diff --git a/docs/conf.py b/docs/conf.py index 49599bfd..7dc4bd3c 100644 --- a/docs/conf.py +++ b/docs/conf.py @@ -227,6 +227,7 @@ def new_process_docstring(app, what, name, obj, options, lines): nb_execution_allow_errors = False nb_execution_excludepatterns = [ # slow examples + 'nanolm.ipynb', 'cifar10_resnet.ipynb', 'adversarial_training.ipynb', 'reduce_on_plateau.ipynb', @@ -296,6 +297,7 @@ def linkcode_resolve(domain, info): intersphinx_mapping = { 'jax': ('https://jax.readthedocs.io/en/latest/', None), + 'flax': ('https://flax.readthedocs.io/en/latest/', None), } source_suffix = ['.rst', '.md', '.ipynb'] diff --git a/docs/gallery.rst b/docs/gallery.rst index c99dc1e6..85cc08b2 100644 --- a/docs/gallery.rst +++ b/docs/gallery.rst @@ -150,6 +150,23 @@
Adversarial training of CNN on MNIST.
+ +.. raw:: html + +
+ +.. only:: html + + .. image:: /images/examples/tiny_shakespeare.png + :alt: Character-level Transformer on Tiny Shakespeare + + :doc:`_collections/examples/nanolm` + +.. raw:: html + +
Character-level Transformer on Tiny Shakespeare.
+
+ .. raw:: html diff --git a/docs/images/examples/tiny_shakespeare.png b/docs/images/examples/tiny_shakespeare.png new file mode 100644 index 00000000..62e9571b Binary files /dev/null and b/docs/images/examples/tiny_shakespeare.png differ diff --git a/examples/nanolm.ipynb b/examples/nanolm.ipynb new file mode 100644 index 00000000..90d22d4d --- /dev/null +++ b/examples/nanolm.ipynb @@ -0,0 +1,733 @@ +{ + "cells": [ + { + "cell_type": "markdown", + "metadata": { + "id": "xpfwcMJHTtfw" + }, + "source": [ + "\n", + "# Character-level Transformer on Tiny Shakespeare\n", + "\n", + "[![Open in Colab](https://colab.research.google.com/assets/colab-badge.svg)](https://colab.sandbox.google.com/github/google-deepmind/optax/blob/main/examples/nanolm.ipynb)\n", + "\n", + "This example demonstrates how to train a small-scale transformer-based language model (inspired by NanoGPT) on the Tiny Shakespeare dataset. The core idea is to train a model that can predict the next character in a sequence of text based on the characters that came before it.\n", + "\n", + "**Why the Tiny Shakespeare Dataset?**\n", + "\n", + "* **Manageable Size:** Since we're building a small-scale model, the Tiny Shakespeare dataset provides a suitable training corpus without overwhelming computational resources.\n", + "* **Linguistic Complexity:** Shakespeare's works offer a rich vocabulary and interesting grammatical patterns, making the dataset a good testbed for our model's language learning abilities.\n", + "* **Accessibility:** Easily accessible through [TensorFlow Datasets](https://www.tensorflow.org/datasets).\n", + "\n", + "**Libraries Used**\n", + "\n", + "* **JAX:** Provides the foundation for numerical computations and automatic differentiation.\n", + "* **Tensorflow Datasets (`tfds`)** Offers easy access to the Tiny Shakespeare dataset.\n", + "* **Flax's Linen Module:** Provides building blocks for defining our neural network architecture.\n", + "* **Optax:** Contains a library of optimization algorithms for training the model's parameters. In this example we'll use the {py:func}`optax.adamw` solver." + ] + }, + { + "cell_type": "code", + "execution_count": null, + "metadata": { + "id": "jIabArrRWFw0", + "outputId": "75c89cb3-35ca-4217-a8e0-4b45f9efe31f", + "colab": { + "base_uri": "https://localhost:8080/" + } + }, + "outputs": [ + { + "output_type": "stream", + "name": "stdout", + "text": [ + "JAX running on GPU\n" + ] + } + ], + "source": [ + "import functools\n", + "\n", + "import flax.linen as nn\n", + "import jax\n", + "import jax.numpy as jnp\n", + "from matplotlib import pyplot as plt\n", + "import optax\n", + "import tensorflow_datasets as tfds\n", + "\n", + "# platform check\n", + "print(\"JAX running on\", jax.devices()[0].platform.upper())" + ] + }, + { + "cell_type": "markdown", + "metadata": { + "id": "UhFD-uojcAI6" + }, + "source": [ + "# Hyperparameters and dataset download\n", + "\n", + "Next, we set some important hyperparameters. This includes hyperparameters for the training process such as the learning rate `LEARNING_RATE` and the batch size `BATCH_SIZE`, as well as model parameters such as the context window size `BLOCK_SIZE` and the number of layers `NUM_LAYERS`.\n", + "\n", + "\n", + "After setting these, we load the Tiny Shakespeare dataset and print the length of the training set, which is around one million characters, and that of the validation set (around 50k characters). Finally, we print a small snippet of the train set." + ] + }, + { + "cell_type": "code", + "execution_count": null, + "metadata": { + "id": "34pKN_bIXt8O" + }, + "outputs": [], + "source": [ + "# @markdown Random seed:\n", + "SEED = 42 # @param{type:\"integer\"}\n", + "# @markdown Learning rate passed to the optimizer:\n", + "LEARNING_RATE = 5e-3 # @param{type:\"number\"}\n", + "# @markdown Batch size:\n", + "BATCH_SIZE = 128 # @param{type:\"integer\"}\n", + "# @markdown Numer of training iterations:\n", + "N_ITERATIONS = 50_000 # @param{type:\"integer\"}\n", + "# @markdown Number of training iterations between two consecutive evaluations:\n", + "N_FREQ_EVAL = 2_000 # @param{type:\"integer\"}\n", + "# @markdown Batch size\n", + "BATCH_SIZE = 512 # @param{type:\"integer\"}\n", + "# @markdown Rate for dropout in the transformer model\n", + "DROPOUT_RATE = 0.2 # @param{type:\"number\"}\n", + "# @markdown Context window for the transformer model\n", + "BLOCK_SIZE = 64 # @param{type:\"integer\"}\n", + "# @markdown Number of layer for the transformer model\n", + "NUM_LAYERS = 6 # @param{type:\"integer\"}\n", + "# @markdown Size of the embedding for the transformer model\n", + "EMBED_SIZE = 256 # @param{type:\"integer\"}\n", + "# @markdown Number of heads for the transformer model\n", + "NUM_HEADS = 8 # @param{type:\"integer\"}\n", + "# @markdown Size of the heads for the transformer model\n", + "HEAD_SIZE = 32 # @param{type:\"integer\"}\n" + ] + }, + { + "cell_type": "code", + "execution_count": null, + "metadata": { + "id": "mghpbB9653Gw" + }, + "outputs": [], + "source": [ + "ds = tfds.load(\"tiny_shakespeare\")\n", + "\n", + "# combine train and test examples into a single string\n", + "text_train = \"\"\n", + "for example in ds[\"train\"].concatenate(ds[\"test\"]).as_numpy_iterator():\n", + " text_train += example[\"text\"].decode(\"utf-8\")\n", + "\n", + "# similarly, create a single string for validation\n", + "text_validation = \"\"\n", + "for example in ds[\"validation\"].as_numpy_iterator():\n", + " text_validation += example[\"text\"].decode(\"utf-8\")" + ] + }, + { + "cell_type": "code", + "execution_count": null, + "metadata": { + "id": "USiJ0GjWSPu_", + "outputId": "bd033bb3-1c6a-4dc4-a1e3-b38a04dd43c0", + "colab": { + "base_uri": "https://localhost:8080/" + } + }, + "outputs": [ + { + "output_type": "stream", + "name": "stdout", + "text": [ + "Length of text for training: 1_059_624 characters\n", + "Length of text for validation: 55_770 characters\n" + ] + } + ], + "source": [ + "print(f\"Length of text for training: {len(text_train):_} characters\")\n", + "print(f\"Length of text for validation: {len(text_validation):_} characters\")" + ] + }, + { + "cell_type": "code", + "execution_count": null, + "metadata": { + "id": "wOq-djQ9cueI", + "outputId": "8d639dad-b1df-4536-8f74-bed86a869201", + "colab": { + "base_uri": "https://localhost:8080/" + } + }, + "outputs": [ + { + "output_type": "stream", + "name": "stdout", + "text": [ + "First Citizen:\n", + "Before we proceed any further, hear me speak.\n", + "\n", + "All:\n", + "Speak, speak.\n", + "\n", + "First Citizen:\n", + "You are all resolved rather to die than to famish?\n", + "\n", + "All:\n", + "Resolved. resolved.\n", + "\n", + "First Citizen:\n", + "First, you know Caius Marcius is chief enemy to the people.\n", + "\n", + "All:\n", + "We know't, we know't.\n", + "\n", + "First Citizen:\n", + "Let us kill him, and we'll have corn at our own price.\n", + "Is't a verdict?\n", + "\n", + "All:\n", + "No more talking on't; let it be done: away, away!\n", + "\n", + "Second Citizen:\n", + "One word, good citizens.\n", + "\n", + "First Citizen:\n", + "We are accounted poor citizens, the patricians good.\n", + "What authority surfeits on would relieve us: if they\n", + "would yield us but the superfluity, while it were\n", + "wholesome, we might guess they relieved us humanely;\n", + "but they think we are too dear: the leanness that\n", + "afflicts us, the object of our misery, is as an\n", + "inventory to particularise their abundance; our\n", + "sufferance is a gain to them Let us revenge this with\n", + "our pikes, ere we become rakes: for the gods know I\n", + "speak this in hunger for bread, not in thirst for revenge.\n", + "\n", + "\n" + ] + } + ], + "source": [ + "# small sample of the train set\n", + "print(text_train[:1000])" + ] + }, + { + "cell_type": "markdown", + "metadata": { + "id": "FguiERfTcEPa" + }, + "source": [ + "# Data preparation\n", + "\n", + "To prepare the data for the model, we first create a vocabulary consisting of all the unique characters in the dataset. We print that vocabulary and its size.\n", + "\n", + "We then define encoding and decoding functions to convert text into sequences of integers (representing our characters) and vice versa.\n", + "\n", + "Finally, we define a function `get_batch` that returns random mini-batches of data. This function uses JAX's {py:func}`jax.lax.dynamic_slice` function to efficiently handle sequences of varying lengths within batches. The `@jax.jit` decorator compiles this function for faster execution. The function randomly samples a batch from the data and prepares input sequences (`x`) and target sequences (`y`). The target sequence is simply the input sequence shifted by one position, as the goal of the language model is to predict the next character given the previous ones.\n" + ] + }, + { + "cell_type": "code", + "execution_count": null, + "metadata": { + "id": "rESkNoDXFE-4", + "outputId": "42579042-716c-4101-cf08-14ebefa6c984", + "colab": { + "base_uri": "https://localhost:8080/" + } + }, + "outputs": [ + { + "output_type": "stream", + "name": "stdout", + "text": [ + "Vocabulary:, \n", + " !$\u0026',-.3:;?ABCDEFGHIJKLMNOPQRSTUVWXYZabcdefghijklmnopqrstuvwxyz\n", + "Length of vocabulary: 65\n" + ] + } + ], + "source": [ + "vocab = sorted(list(set(text_train)))\n", + "print(\"Vocabulary:, \", \"\".join(vocab))\n", + "print(\"Length of vocabulary: \", len(vocab))" + ] + }, + { + "cell_type": "code", + "execution_count": null, + "metadata": { + "id": "F-LSTr86bXrV" + }, + "outputs": [], + "source": [ + "# create a mapping from characters to integers\n", + "stoi = {ch: i for i, ch in enumerate(vocab)}\n", + "itos = {i: ch for i, ch in enumerate(vocab)}\n", + "encode = lambda s: [\n", + " stoi[c] for c in s\n", + "] # encoder: take a string, output a list of integers\n", + "decode = lambda l: \"\".join(\n", + " [itos[i] for i in l]\n", + ") # decoder: take a list of integers, output a string\n", + "\n", + "# encode train and validation data\n", + "train_data = jnp.array(encode(text_train))\n", + "eval_data = jnp.array(encode(text_validation))" + ] + }, + { + "cell_type": "code", + "execution_count": null, + "metadata": { + "id": "tZLP1asmb8WY" + }, + "outputs": [], + "source": [ + "dynamic_slice_vmap = jax.vmap(jax.lax.dynamic_slice, in_axes=(None, 0, None))\n", + "\n", + "\n", + "@jax.jit\n", + "def get_batch(random_key, data):\n", + " \"\"\"Prepares a random batch of training data.\n", + "\n", + " Args:\n", + " random_key: A random seed for sampling a batch.\n", + " data: The complete training dataset.\n", + "\n", + " Returns:\n", + " x: Input sequences.\n", + " y: Target sequences (shifted inputs).\n", + " \"\"\"\n", + " ix = jax.random.randint(\n", + " random_key, shape=(BATCH_SIZE, 1), minval=0, maxval=len(data) - BLOCK_SIZE\n", + " )\n", + " x = dynamic_slice_vmap(data, ix, (BLOCK_SIZE,))\n", + " y = dynamic_slice_vmap(data, ix + 1, (BLOCK_SIZE,))\n", + " return x, y" + ] + }, + { + "cell_type": "markdown", + "metadata": { + "id": "PQfTaf2UcTSc" + }, + "source": [ + "# NanoLM Model Definition\n", + "\n", + "The NanoLM model itself is defined as a Flax Linen module. The core of the model is a Transformer architecture, designed for sequence-to-sequence tasks like language modeling. Key parameters of the model, such as the number of layers, attention heads, and embedding size, are specified here.\n", + "\n", + "Inside the model's `__call__` method, we first embed our input characters into vector representations. Positional embeddings are added to provide the model with a sense of order in the sequence. The core of the Transformer consists of multiple layers. Each layer has two main components:\n", + "\n", + " * **Multi-Head Attention**: This mechanism allows the model to \"attend\" to different parts of the input sequence, improving its understanding of context and relationships within the text. In the code this is implemented through the {py:class}`flax.linen.MultiHeadDotProductAttention` class.\n", + "\n", + " * **Feedforward Network**: This network processes the output of the attention layer, applying non-linear transformations to further learn complex patterns in the data. This is implemented through the {py:class}`flax.linen.Sequential` class.\n", + "\n", + "Normalization and dropout (for regularization) are used within the layers to improve training stability. Finally, a dense layer maps the model's output to the vocabulary size, producing probabilities for each character as the next potential character.\n", + "\n", + "The generate function enables the model to create new text sequences. It iteratively generates one character at a time, conditioned on the previously generated text.\n" + ] + }, + { + "cell_type": "code", + "execution_count": null, + "metadata": { + "id": "-c7M35UaYMyD" + }, + "outputs": [], + "source": [ + "class NanoLM(nn.Module):\n", + " \"\"\"NanoLM model.\"\"\"\n", + " vocab_size: int\n", + " num_layers: int = 6\n", + " num_heads: int = 8\n", + " head_size: int = 32\n", + " dropout_rate: float = 0.2\n", + " embed_size: int = 256\n", + " block_size: int = 64\n", + "\n", + " @nn.compact\n", + " def __call__(self, x, training: bool):\n", + " seq_len = x.shape[1]\n", + "\n", + " x = nn.Embed(self.vocab_size, self.embed_size)(x) + nn.Embed(\n", + " self.block_size, self.embed_size\n", + " )(jnp.arange(seq_len))\n", + " for _ in range(self.num_layers):\n", + " x_norm = nn.LayerNorm()(x)\n", + " x = x + nn.MultiHeadDotProductAttention(\n", + " num_heads=self.num_heads,\n", + " qkv_features=self.head_size,\n", + " out_features=self.head_size * self.num_heads,\n", + " dropout_rate=self.dropout_rate,\n", + " )(\n", + " x_norm,\n", + " x_norm,\n", + " mask=jnp.tril(jnp.ones((x.shape[-2], x.shape[-2]))),\n", + " deterministic=not training,\n", + " )\n", + "\n", + " x = x + nn.Sequential([\n", + " nn.Dense(4 * self.embed_size),\n", + " nn.relu,\n", + " nn.Dropout(self.dropout_rate, deterministic=not training),\n", + " nn.Dense(self.embed_size),\n", + " ])(nn.LayerNorm()(x))\n", + "\n", + " x = nn.LayerNorm()(x)\n", + " return nn.Dense(self.vocab_size)(x)\n", + "\n", + " @functools.partial(jax.jit, static_argnames=(\"self\", \"length\"))\n", + " def generate(self, rng, params, length):\n", + " def _scan_generate(carry, _):\n", + " random_key, context = carry\n", + " logits = self.apply(params, context, training=False)\n", + " rng, rng_subkey = jax.random.split(random_key)\n", + " new_token = jax.random.categorical(\n", + " rng_subkey, logits[:, -1, :], axis=-1, shape=(1, 1)\n", + " )\n", + " context = jnp.concatenate([context[:, 1:], new_token], axis=1)\n", + " return (rng, context), new_token\n", + "\n", + " _, new_tokens = jax.lax.scan(\n", + " _scan_generate,\n", + " (rng, jnp.zeros((1, self.block_size), dtype=jnp.int32)),\n", + " (),\n", + " length=length,\n", + " )\n", + " return new_tokens" + ] + }, + { + "cell_type": "markdown", + "metadata": { + "id": "ylVNKEhscy9d" + }, + "source": [ + "# State, Optimizer, and Loss Definition\n", + "\n", + "This section initializes the model's parameters, defines the loss function used for language modeling, and sets up the training and evaluation processes.\n", + "\n", + "In this case the loss function `loss_fun` is the cross-entropy. It uses dropout for regularization, introduced via the `rngs={\"dropout\": dropout_key}` argument. We also define a function for evaluating the model's performance on unseen data (`eval_step`)." + ] + }, + { + "cell_type": "code", + "execution_count": null, + "metadata": { + "id": "sjSnK3yDYIus" + }, + "outputs": [], + "source": [ + "model = NanoLM(\n", + " vocab_size=len(vocab),\n", + " num_layers=NUM_LAYERS,\n", + " num_heads=NUM_HEADS,\n", + " head_size=HEAD_SIZE,\n", + " dropout_rate=DROPOUT_RATE,\n", + " embed_size=EMBED_SIZE,\n", + " block_size=BLOCK_SIZE,\n", + ")\n", + "\n", + "def loss_fun(params, x, y, dropout_key):\n", + " logits = model.apply(params, x, training=True, rngs={\"dropout\": dropout_key})\n", + " return optax.softmax_cross_entropy_with_integer_labels(\n", + " logits=logits, labels=y\n", + " ).mean()\n", + "\n", + "\n", + "@jax.jit\n", + "def eval_step(params, x, y):\n", + " logits = model.apply(params, x, training=False)\n", + " return optax.softmax_cross_entropy_with_integer_labels(\n", + " logits=logits, labels=y\n", + " ).mean()" + ] + }, + { + "cell_type": "code", + "execution_count": null, + "metadata": { + "id": "ejU1Yt8XIH80" + }, + "outputs": [], + "source": [ + "key = jax.random.PRNGKey(SEED)\n", + "key, subkey = jax.random.split(key)\n", + "\n", + "var_params = model.init(\n", + " key,\n", + " jnp.ones((BATCH_SIZE, BLOCK_SIZE), dtype=jnp.int32),\n", + " training=False,\n", + ")" + ] + }, + { + "cell_type": "markdown", + "metadata": { + "id": "kgSjWONs4eFp" + }, + "source": [ + "We've now instatiated a NanoLM model with the following number of parameters" + ] + }, + { + "cell_type": "code", + "execution_count": null, + "metadata": { + "id": "9Ckqdkd6QVsl", + "outputId": "2c202753-b938-4148-992c-cc63077be607", + "colab": { + "base_uri": "https://localhost:8080/" + } + }, + "outputs": [ + { + "output_type": "stream", + "name": "stdout", + "text": [ + "Total number of parameters: 3_408_513\n" + ] + } + ], + "source": [ + "n_params = sum(p.size for p in jax.tree_util.tree_leaves(var_params))\n", + "\n", + "print(f\"Total number of parameters: {n_params:_}\")" + ] + }, + { + "cell_type": "markdown", + "metadata": { + "id": "vreyB_oo4Zch" + }, + "source": [ + "# Model training\n", + "\n", + "We start by creating an optimizer and instantiating its state. In this case we'll use {py:func}`optax.adamw` but it's possible to replace this with other optax optimizers.\n", + "\n", + "\n", + "We then proceeded to the training loop. For maximum efficiency we extracted the most computationally intensive tasks inside the `step` function and just-in-time compile this function using `@jax.jit`. This allows JAX to perform some optimizations in our code and generally achieve a much higher efficiency than without.\n", + "\n", + "Inside the training loop, we call the aforementioned `step` functions, as well as computing accuracy on a validation set every `N_FREQ_EVAL` iterations." + ] + }, + { + "cell_type": "code", + "execution_count": null, + "metadata": { + "id": "1xwLpjDxccMi" + }, + "outputs": [], + "source": [ + "# To run with SGD instead of adam, replace `adam` with `sgd`\n", + "opt = optax.adamw(learning_rate=LEARNING_RATE)\n", + "\n", + "opt_state = opt.init(var_params)" + ] + }, + { + "cell_type": "code", + "execution_count": null, + "metadata": { + "colab": { + "base_uri": "https://localhost:8080/" + }, + "id": "DhnK0G7AQUCA", + "outputId": "37b5a3bd-c0bc-47c2-a828-62356ab4cbe5" + }, + "outputs": [ + { + "output_type": "stream", + "name": "stdout", + "text": [ + "Step: 0\t train loss: 4.586461067199707\t eval loss: 6.037547588348389\n", + "Step: 2000\t train loss: 1.3876519203186035\t eval loss: 1.4252502918243408\n", + "Step: 4000\t train loss: 1.280821442604065\t eval loss: 1.4000993967056274\n", + "Step: 6000\t train loss: 1.1978516578674316\t eval loss: 1.4045076370239258\n", + "Step: 8000\t train loss: 1.177159070968628\t eval loss: 1.387284278869629\n", + "Step: 10000\t train loss: 1.1472305059432983\t eval loss: 1.423332929611206\n", + "Step: 12000\t train loss: 1.107376217842102\t eval loss: 1.4390857219696045\n", + "Step: 14000\t train loss: 1.096900224685669\t eval loss: 1.474606990814209\n", + "Step: 16000\t train loss: 1.0772775411605835\t eval loss: 1.4595460891723633\n", + "Step: 18000\t train loss: 1.050074577331543\t eval loss: 1.4540534019470215\n", + "Step: 20000\t train loss: 1.0540519952774048\t eval loss: 1.4794009923934937\n", + "Step: 22000\t train loss: 1.035787582397461\t eval loss: 1.5094046592712402\n", + "Step: 24000\t train loss: 1.0402700901031494\t eval loss: 1.5403311252593994\n", + "Step: 26000\t train loss: 1.0363699197769165\t eval loss: 1.5168808698654175\n", + "Step: 28000\t train loss: 1.0000224113464355\t eval loss: 1.5624736547470093\n", + "Step: 30000\t train loss: 0.9905486106872559\t eval loss: 1.5720288753509521\n", + "Step: 32000\t train loss: 0.9986284971237183\t eval loss: 1.526949167251587\n", + "Step: 34000\t train loss: 0.9822986125946045\t eval loss: 1.5732147693634033\n", + "Step: 36000\t train loss: 0.9837050437927246\t eval loss: 1.6341228485107422\n", + "Step: 38000\t train loss: 0.9723658561706543\t eval loss: 1.542256474494934\n", + "Step: 40000\t train loss: 0.9632444977760315\t eval loss: 1.5658419132232666\n", + "Step: 42000\t train loss: 0.9664344787597656\t eval loss: 1.6112396717071533\n", + "Step: 44000\t train loss: 0.9476494789123535\t eval loss: 1.6128767728805542\n", + "Step: 46000\t train loss: 0.9473679065704346\t eval loss: 1.5689520835876465\n", + "Step: 48000\t train loss: 0.9642226696014404\t eval loss: 1.5935924053192139\n", + "CPU times: user 15min 19s, sys: 9min 8s, total: 24min 27s\n", + "Wall time: 22min 48s\n" + ] + } + ], + "source": [ + "%%time\n", + "\n", + "all_train_losses = []\n", + "all_eval_losses = []\n", + "\n", + "# we define one iteration of the optimizer and JIT this function\n", + "@jax.jit\n", + "def step(key, params, opt_state):\n", + " key, subkey = jax.random.split(key)\n", + " batch = get_batch(key, train_data)\n", + " loss, grad = jax.value_and_grad(loss_fun)(params, *batch, subkey)\n", + " updates, opt_state = opt.update(grad, opt_state, params)\n", + " params = optax.apply_updates(params, updates)\n", + " return params, key, opt_state, loss\n", + "\n", + "\n", + "for i in range(N_ITERATIONS):\n", + " var_params, key, opt_state, loss = step(key, var_params, opt_state)\n", + " all_train_losses.append(loss)\n", + "\n", + " # once every N_FREQ_EVAL we compute loss on the validation set\n", + " if i % N_FREQ_EVAL == 0:\n", + " key, subkey = jax.random.split(key)\n", + " eval_loss = eval_step(var_params, *get_batch(subkey, eval_data))\n", + " all_eval_losses.append(eval_loss)\n", + " print(f\"Step: {i}\\t train loss: {loss}\\t eval loss: {eval_loss}\")" + ] + }, + { + "cell_type": "code", + "execution_count": null, + "metadata": { + "colab": { + "base_uri": "https://localhost:8080/", + "height": 472 + }, + "id": "Gc-V4kAKAA9q", + "outputId": "68b59bf2-9621-4c2e-e0e0-be665d19630e" + }, + "outputs": [ + { + "output_type": "display_data", + "data": { + "text/plain": [ + "\u003cFigure size 640x480 with 1 Axes\u003e" + ], + "image/png": "iVBORw0KGgoAAAANSUhEUgAAAioAAAHHCAYAAACRAnNyAAAAOXRFWHRTb2Z0d2FyZQBNYXRwbG90bGliIHZlcnNpb24zLjcuMSwgaHR0cHM6Ly9tYXRwbG90bGliLm9yZy/bCgiHAAAACXBIWXMAAA9hAAAPYQGoP6dpAABg/ElEQVR4nO3dd3hUxcIG8PdsTd30QgoJJJAQCL0YkaI0ARHUq4iogP2KnyK2q9cCooJ4RbGhiIJy5WKjqYAGIRSld5BOKNJCgGRTN5vd+f5Yssm2ZBO2Bd7f8+yTnLOz58xOEvZlZs4cSQghQEREROSDZN6uABEREZEjDCpERETksxhUiIiIyGcxqBAREZHPYlAhIiIin8WgQkRERD6LQYWIiIh8FoMKERER+SwGFSIiIvJZDCpE5HLLly9H+/bt4efnB0mSUFBQ4JbzzJkzB5Ik4dixY245vqc9/vjj6Nevn8fON3r0aCQnJ7vseDk5OZAkCTk5OS47Zn1cuHABgYGBWLp0qVfOT+7BoEIeceTIETz66KNo3rw5/Pz8oNFo0L17d0yfPh1lZWXerh650IULF3DXXXfB398fH3/8MebOnYvAwEBvV8vn5ebmYtasWXjppZfM+06fPo0JEyZgx44d3qtYIxIREYGHHnoIr7zyirerQi6k8HYF6Or3yy+/4M4774Rarcb999+PNm3aoKKiAuvWrcNzzz2HvXv3YubMmd6uJrnI5s2bUVRUhEmTJqFv377erk6jMX36dDRr1gw33nijed/p06cxceJEJCcno3379i4/5+effw6j0ejy43rTY489hg8++AArV67ETTfd5O3qkAswqJBb5ebm4u6770ZSUhJWrlyJJk2amJ8bO3YsDh8+jF9++cWLNbxy5eXlUKlUkMnYQQkAeXl5AIDQ0FDvVqQR0ev1+Oabb/DYY49d0XFKS0sREBDgdHmlUnlF5/NFrVq1Qps2bTBnzhwGlasE/2Ult5o6dSqKi4vxxRdfWISUKqmpqXjqqafM25WVlZg0aRJSUlKgVquRnJyMl156CTqdzuJ1ycnJuOWWW7Bu3Tp07doVfn5+aN68Ob7++mtzmS1btkCSJHz11Vc25/31118hSRJ+/vln875Tp07hgQceQExMDNRqNVq3bo0vv/zS4nVVY/Dz58/Hyy+/jPj4eAQEBECr1QIAvv/+e2RkZMDPzw9t2rTBwoUL7c4DMBqNeP/999G6dWv4+fkhJiYGjz76KC5dulTv91mloKAATz/9NJKTk6FWq5GQkID7778f+fn55jI6nQ6vvfYaUlNToVarkZiYiOeff96mfR35/vvv0alTJ/j7+yMyMhL33nsvTp06ZX6+d+/eGDVqFACgS5cukCQJo0ePdni848eP4/HHH0daWhr8/f0RERGBO++80+6ck7179+Kmm26Cv78/EhIS8MYbb9jtDVi8eDEGDx6MuLg4qNVqpKSkYNKkSTAYDBblevfujTZt2mDXrl3o1asXAgICkJqaih9++AEAsHr1anTr1g3+/v5IS0vDihUrzK/dtWsXJEnCkiVLzPu2bt0KSZLQsWNHi/MMHDgQ3bp1c9yoANatW4f8/HyLHqicnBx06dIFADBmzBhIkgRJkjBnzhyL+m/duhU9e/ZEQECAedjI2Taw/t08duwYJEnCf/7zH8ycOdP8d9ilSxds3ry51vdQm7p+bwDg7NmzGDNmDBISEqBWq9GkSRMMHTrU4ndhy5YtGDBgACIjI+Hv749mzZrhgQcesDlfv3798NNPP0EI0eA6kw8RRG4UHx8vmjdv7nT5UaNGCQDiH//4h/j444/F/fffLwCIYcOGWZRLSkoSaWlpIiYmRrz00kvio48+Eh07dhSSJIk9e/aYyzVv3lwMGjTI5jxjxowRYWFhoqKiQgghxNmzZ0VCQoJITEwUr7/+upgxY4a49dZbBQDx3nvvmV+3atUqAUBkZGSI9u3bi2nTponJkyeLkpIS8fPPPwtJkkTbtm3FtGnTxCuvvCLCwsJEmzZtRFJSksX5H3roIaFQKMTDDz8sPv30U/HCCy+IwMBA0aVLF3Od6vM+i4qKRJs2bYRcLhcPP/ywmDFjhpg0aZLo0qWL2L59uxBCCIPBIPr37y8CAgLEuHHjxGeffSaeeOIJoVAoxNChQ+v82cyePVsAEF26dBHvvfee+Ne//iX8/f1FcnKyuHTpkhBCiN9++0088sgjAoB4/fXXxdy5c8Wff/7p8Jjff/+9aNeunXj11VfFzJkzxUsvvSTCwsJEUlKSKCkpMZc7c+aMiIqKEmFhYWLChAninXfeES1atBBt27YVAERubq657LBhw8Rdd90l3nnnHTFjxgxx5513CgDi2WeftTh3r169RFxcnEhMTBTPPfec+PDDD0VGRoaQy+Vi/vz5IjY2VkyYMEG8//77Ij4+XoSEhAitVmtuy9DQUPHMM8+Yj/fee+8JmUwmZDKZKCwsNJfTaDQ257b2xhtvCEmSzK8TwvQ7+frrrwsA4pFHHhFz584Vc+fOFUeOHDHXPzY2VkRFRYn/+7//E5999plYtGhRvdpg1KhRFr+bubm5AoDo0KGDSE1NFW+//baYOnWqiIyMFAkJCRa/m/ZU/X2sWrXKvM+Z3xshhLj++utFSEiIePnll8WsWbPEW2+9JW688UaxevVqIYQQ586dE2FhYaJly5binXfeEZ9//rn497//LVq1amVTj//+978CgNi9e3et9aXGgUGF3KawsFAAcOpDUAghduzYIQCIhx56yGL/s88+KwCIlStXmvclJSUJAGLNmjXmfXl5eUKtVlt8eLz44otCqVSKixcvmvfpdDoRGhoqHnjgAfO+Bx98UDRp0kTk5+dbnPvuu+8WISEhorS0VAhR/Q9x8+bNzfuqZGZmioSEBFFUVGTel5OTIwBYfBisXbtWABDffPONxeuXL19us9/Z9/nqq68KAGLBggXCmtFoFEIIMXfuXCGTycTatWstnv/0008FAPHHH3/YvLZKRUWFiI6OFm3atBFlZWXm/T///LMAIF599VXzvqoPps2bNzs8XhXrNhRCiPXr1wsA4uuvvzbvGzdunAAgNm7caN6Xl5cnQkJCbIKKvWM++uijIiAgQJSXl5v39erVSwAQ8+bNM+/bv3+/ACBkMpnYsGGDef+vv/4qAIjZs2eb9w0ePFh07drVvH377beL22+/XcjlcrFs2TIhhBDbtm0TAMTixYtrbYd7771XRERE2OzfvHmzzXmt6//pp5/aPOdsGzgKKhERERZ/M4sXLxYAxE8//VTr+7AOKs7+3ly6dEkAEO+8847DYy9cuNDp36s///xTABDffvttnWXJ93Hoh9ymajgkODjYqfJVlxSOHz/eYv8zzzwDADZzWTIyMtCjRw/zdlRUFNLS0nD06FHzvuHDh0Ov12PBggXmfb/99hsKCgowfPhwAIAQAj/++COGDBkCIQTy8/PNjwEDBqCwsBDbtm2zOPeoUaPg7+9v3j59+jR2796N+++/H0FBQeb9vXr1QmZmpsVrv//+e4SEhKBfv34W5+rUqROCgoKwatWqer/PH3/8Ee3atcNtt91m066SJJnP26pVK6Snp1uct2oc3/q8NW3ZsgV5eXl4/PHH4efnZ94/ePBgpKenN3ieUc021Ov1uHDhAlJTUxEaGmrR5kuXLsV1112Hrl27mvdFRUVh5MiRtR6zqKgI+fn56NGjB0pLS7F//36LskFBQbj77rvN22lpaQgNDUWrVq0shmuqvq/Z5j169MC2bdtQUlICwDR8M2jQILRv3x5r164FAKxduxaSJOGGG26otR0uXLiAsLCwWsvYo1arMWbMGJv99WkDe4YPH25Rn6rfv5rv3xnO/t74+/tDpVIhJyfHZvizStWcp59//hl6vb7W81bVveawJzVeDCrkNhqNBoDpH0pnHD9+HDKZDKmpqRb7Y2NjERoaiuPHj1vsb9q0qc0xwsLCLP6ha9euHdLT0/Htt9+a93377beIjIw0f0CfP38eBQUFmDlzJqKioiweVR8CVRNEqzRr1sym7gBs6m5v36FDh1BYWIjo6Gib8xUXF9ucy5n3eeTIEbRp08amnPV59+7da3POli1b2n2P9t5fWlqazXPp6ek2PxtnlZWV4dVXX0ViYiLUajUiIyMRFRWFgoICFBYWWpy/RYsWNq+3V5+9e/fitttuQ0hICDQaDaKionDvvfcCgMUxASAhIcEc5KqEhIQgMTHRZh8Aizbv0aMHKisrsX79ehw4cAB5eXno0aMHevbsaRFUMjIyEB4eXmdbiAbMp4iPj4dKpbLZX582sMf6d67qg99RiHDE2d8btVqNt99+G8uWLUNMTAx69uyJqVOn4uzZs+byvXr1wh133IGJEyciMjISQ4cOxezZs+3Or6pqS+ufLTVOvOqH3Eaj0SAuLg579uyp1+uc/cdFLpfb3W/9D/7w4cPx5ptvIj8/H8HBwViyZAlGjBgBhcL06181IfPee+81TwS11rZtW4vtmv9jrS+j0Yjo6Gh88803dp+Pioqy2Hb2fTpz3szMTEybNs3u89Yfzp7wf//3f5g9ezbGjRuHrKwshISEQJIk3H333Q26bLagoAC9evWCRqPB66+/jpSUFPj5+WHbtm144YUXbI7pqG2dafPOnTvDz88Pa9asQdOmTREdHY2WLVuiR48e+OSTT6DT6bB27Vq7vVzWIiIi6h0CAPu/h/VtA3tc9TtXH+PGjcOQIUOwaNEi/Prrr3jllVcwefJkrFy5Eh06dIAkSfjhhx+wYcMG/PTTT/j111/xwAMP4N1338WGDRssejKr2jIyMtJt9SXPYVAht7rlllswc+ZMrF+/HllZWbWWTUpKgtFoxKFDh9CqVSvz/nPnzqGgoABJSUkNqsPw4cMxceJE/Pjjj4iJiYFWq7Xo7o+KikJwcDAMBkOD1/2oqtvhw4dtnrPel5KSghUrVqB79+5XFHisj1lXIExJScHOnTvRp0+fev9Ps+r9HThwwOaSzwMHDjT4Z/PDDz9g1KhRePfdd837ysvLbVayTUpKwqFDh2xef+DAAYvtnJwcXLhwAQsWLEDPnj3N+3NzcxtUv9qoVCp07doVa9euRdOmTc3DIz169IBOp8M333yDc+fOWdTDkfT0dHzzzTcoLCw0994ADesR8GQb1KW+vzcpKSl45pln8Mwzz+DQoUNo37493n33Xfz3v/81l7nuuutw3XXX4c0338S8efMwcuRIzJ8/Hw899JC5TNV7rfnvCDVeHPoht3r++ecRGBiIhx56COfOnbN5/siRI5g+fToAYNCgQQCA999/36JMVQ/A4MGDG1SHVq1aITMzE99++y2+/fZbNGnSxOIfcLlcjjvuuAM//vij3Q/78+fP13mOuLg4tGnTBl9//TWKi4vN+1evXo3du3dblL3rrrtgMBgwadIkm+NUVlY2aLn5O+64Azt37sTChQttnqv6X/Bdd92FU6dO4fPPP7cpU1ZWZp5rYU/nzp0RHR2NTz/91KKrfdmyZdi3b1+DfzZyudzmf+kffvihzWW0gwYNwoYNG7Bp0ybzvvPnz9v0SlX1BNQ8ZkVFBT755JMG1a8uPXr0wMaNG7Fq1SpzUImMjESrVq3w9ttvm8vUJSsrC0IIbN261WJ/1Yq+9fmd8HQb1MbZ35vS0lKUl5dbvDYlJQXBwcHm1126dMnmd6VqETzr4Z+tW7ciJCQErVu3dvVbIi9gjwq5VUpKCubNm4fhw4ejVatWFivT/vnnn/j+++/N62y0a9cOo0aNwsyZM83d15s2bcJXX32FYcOGWazYWV/Dhw/Hq6++Cj8/Pzz44IM2i7NNmTIFq1atQrdu3fDwww8jIyMDFy9exLZt27BixQpcvHixznO89dZbGDp0KLp3744xY8bg0qVL+Oijj9CmTRuL8NKrVy88+uijmDx5Mnbs2IH+/ftDqVTi0KFD+P777zF9+nT84x//qNf7e+655/DDDz/gzjvvxAMPPIBOnTrh4sWLWLJkCT799FO0a9cO9913H7777js89thjWLVqFbp37w6DwYD9+/fju+++w6+//orOnTvbPb5SqcTbb7+NMWPGoFevXhgxYgTOnTuH6dOnIzk5GU8//XS96lvllltuwdy5cxESEoKMjAysX78eK1asQEREhEW5559/HnPnzsXNN9+Mp556CoGBgZg5cyaSkpKwa9cuc7nrr78eYWFhGDVqFJ588klIkoS5c+e6bciiR48eePPNN3Hy5EmLQNKzZ0989tlnSE5ORkJCQp3HueGGGxAREYEVK1ZY9DykpKQgNDQUn376KYKDgxEYGIhu3brZzJGqydNtUBtnf28OHjyIPn364K677kJGRgYUCgUWLlyIc+fOmXs/v/rqK3zyySe47bbbkJKSgqKiInz++efQaDTm/+RUyc7OxpAhQzhH5Wrh+QuN6Fp08OBB8fDDD4vk5GShUqlEcHCw6N69u/jwww8tLpfU6/Vi4sSJolmzZkKpVIrExETx4osvWpQRwnTZ7uDBg23O06tXL9GrVy+b/YcOHRIABACxbt06u3U8d+6cGDt2rEhMTBRKpVLExsaKPn36iJkzZ5rLVF1++f3339s9xvz580V6erpQq9WiTZs2YsmSJeKOO+4Q6enpNmVnzpwpOnXqJPz9/UVwcLDIzMwUzz//vDh9+nSD3ueFCxfEE088IeLj44VKpRIJCQli1KhRFpdcV1RUiLffflu0bt1aqNVqERYWJjp16iQmTpxosYaHI99++63o0KGDUKvVIjw8XIwcOVL8/fffFmXqc3nypUuXxJgxY0RkZKQICgoSAwYMEPv37xdJSUli1KhRFmV37dolevXqJfz8/ER8fLyYNGmS+OKLL2wuT/7jjz/EddddJ/z9/UVcXJx4/vnnzZcX11zfo1evXqJ169Y2dXLU5gDE2LFjLfZptVohl8tFcHCwqKysNO+vWsfjvvvuq7MNqjz55JMiNTXVZv/ixYtFRkaGUCgUFpcqO6q/EM63gaPLk+1dJgxAvPbaa7W+B3vrqAhR9+9Nfn6+GDt2rEhPTxeBgYEiJCREdOvWTXz33XfmMtu2bRMjRowQTZs2FWq1WkRHR4tbbrlFbNmyxeJc+/btEwDEihUraq0rNR6SEFy6j8id2rdvj6ioKGRnZ3u7KuTDjh49ivT0dCxbtgx9+vTxdnUarXHjxmHNmjXmlYKp8eMcFSIX0ev1qKystNiXk5ODnTt3onfv3t6pFDUazZs3x4MPPogpU6Z4uyqN1oULFzBr1iy88cYbDClXEfaoELnIsWPH0LdvX9x7772Ii4vD/v378emnnyIkJAR79uyxmXdBRER142RaIhcJCwtDp06dMGvWLJw/fx6BgYEYPHgwpkyZwpBCRNRA7FEhIiIin8U5KkREROSzGFSIiIjIZzXqOSpGoxGnT59GcHAwZ3gTERE1EkIIFBUVIS4uzmYBTmuNOqicPn3aKzdSIyIioit38uTJOldvbtRBJTg4GIDpjWo0GpceW6/X47fffjMvb07uwXb2DLazZ7CdPYPt7DnuamutVovExETz53htGnVQqRru0Wg0bgkqAQEB0Gg0/ENwI7azZ7CdPYPt7BlsZ89xd1s7M22Dk2mJiIjIZzGoEBERkc9iUCEiIiKfxaBCREREPotBhYiIiHwWgwoRERH5LAYVIiIi8lkMKkREROSzGFSIiIjIZzGoEBERkc9iUCEiIiIbycnJ+OCDD7xdDQYVG5U6oOgskH8IYSVHgLy/vF0jIiIip/Tu3Rvjxo1zybE2b96Mhx56yCXHuhJeDyqnTp3Cvffei4iICPj7+yMzMxNbtmzxXoX+/BB4Nw3Kz7LQ8+BEyFe94b26EBERuZAQApWVlU6VjYqKQkBAgJtrVDevBpVLly6he/fuUCqVWLZsGf766y+8++67CAsL816l/EIst3Va79SDiIi8zmgUuFCs8+rDaBRO1XX06NFYvXo1pk+fDkmSIEkS5syZA0mSsGzZMnTq1AlqtRrr1q3DkSNHMHToUMTExCAoKAhdunTBihUrLI5nPfQjSRJmzZqF2267DQEBAWjRogWWLFni0va2R+H2M9Ti7bffRmJiImbPnm3e16xZMy/WCDZBRdIVeakiRETkbZdKK9DpjRV1F3SjrS/3RUSQus5y06dPx8GDB9GmTRu8/vrrAIC9e/cCAP71r3/hP//5D5o3b46wsDCcPHkSgwYNwptvvgm1Wo2vv/4aQ4YMwYEDB9C0aVOH55g4cSKmTp2Kd955Bx9++CFGjhyJ48ePIzw83DVv1g6vBpUlS5ZgwIABuPPOO7F69WrEx8fj8ccfx8MPP2y3vE6ng06nM29rtabeDr1eD71e75I6SYoAi0YxlhXC6KJjk62qn5urfn5kH9vZM9jOnuHJdtY7OUzi7jro9XUPgAQEBECpVMLPzw8REREATEM9APDqq6+id+/e5rIZGRnIyMgwb7/66qtYsGABFi5ciMcff9y832AwmOpwua3vu+8+/OMf/wBgCi0ffPAB/vzzTwwYMKB+76kePzuvBpWjR49ixowZGD9+PF566SVs3rwZTz75JFQqFUaNGmVTfvLkyZg4caLN/t9++81l42jyiwdwS41tXfFFZC9d6pJjk2PZ2dnersI1ge3sGWxnz/BEOxfrAS9/VOL3FSsQpHSu7IULF5Cbm4ullz+3du/eDQAoLCw07wOAsrIyzJ8/H1u3bsXFixdhNBpRUVGBnJwcJCcnAwBKS0tx4MABpKWlmdu6srLS4jgBAQFYsWKFOdA4q7S01OmyXm19o9GIzp0746233gIAdOjQAXv27MGnn35qN6i8+OKLGD9+vHlbq9UiMTER/fv3h0ajcUmdps7VWQQVP2MZBg28GZC8Pu/4qqTX65GdnY1+/fpBqXTyL5Hqje3sGWxnz/BkO18oqcC/t+S49Rx16dO3LyICVU6VnTZtGpo1a4ZBgwYBAAIDAwEAQ4YMQWhoqLnc2LFjsXv3bkybNg0pKSnw9/fH3XffjcTERPNrAwICkJaWBgDo168fAKBr167m5wFAqVQiMzPTYp8zqkZEnOHVoNKkSROLricAaNWqFX788Ue75dVqNdRq23E6pVLpsl/Wv8stjy+TBGRGHeDnmiBE9rnyZ0iOsZ09g+3sGZ5o52iNAltf7uvWc9QlLEAFmUxyqqxarYYQwtwuCoXpY966rdavX4/Ro0fjzjvvBAAUFxfj+PHjkMlkFuXkcrn59VXHs25zuVxe759Dfcp7Nah0794dBw4csNh38OBBJCUlealGQKkUaLtTp2VQISK6BslkklMTWX1FcnIyNm7ciGPHjiEoKAhGo9FuuRYtWmDBggUYMmQIJEnCK6+84rCst3l1POPpp5/Ghg0b8NZbb+Hw4cOYN28eZs6cibFjx3qtTuWSP4zCKrmWF3qnMkRERPXw7LPPQi6XIyMjA1FRUThx4oTdctOmTUNYWBiuv/56DBkyBAMGDEDHjh09XFvneLVHpUuXLli4cCFefPFFvP7662jWrBnef/99jBw50mt1EpIMxfCDBmXVO8u5lgoREfm+li1bYv369Rb7Ro8ebVMuOTkZK1eutNhn3Ulw7Ngx6PV68+TZqiuIaiooKLiyCjvBu1OZAdxyyy245ZZb6i7oQVoEWgYVLvpGRETkFbyUxY4i4W+5gz0qREREXsGgYkcRrNZkKS/wSj2IiIiudQwqdhQJq6DCoR8iIiKvYFCxQ2vTo8KgQkRE5A0MKnawR4WIiMg3MKhYkQAUwXoyLddRISIi8gYGFTtselQ49ENEROQVDCp22Fz1w6EfIiIir2BQsUPLHhUiIiKfwKBiB3tUiIioMerduzfGjRvnsuM9+OCDeOutt1x2vIZgULHDtkeFk2mJiIi8gUHFDpselYpiwGjwTmWIiMh7jEagJN+7D6PRqaqOHj0aq1evxvTp0yFJEiRJwrFjx7Bnzx4MHDgQQUFBiImJwX333Yf8/Hzz63744QdkZmbC398fERER6Nu3L0pKSjBhwgTMnTsXmzZtgkqlgiRJyMnJcVNDO+b1mxL6IpseFcA0/OMf5vnKEBGR95RdBN5J8W4dnjsCBEbWWWz69Ok4ePAg2rRpg9dffx0AoFQq0bVrVzz00EN47733UFZWhhdeeAF33XUXVq5ciTNnzmDEiBGYOnUqbrvtNhQVFWHt2rUQQuDZZ5/F3r17cfToUSxatAhKpRLh4eHufrc2GFTssOlRAUwTahlUiIjIR4WEhEClUiEgIACxsbEAgDfeeAMdOnSwmGfy5ZdfIjExEQcPHkRxcTEqKytx++23IykpCQCQmZlpLuvv7w+FQoHY2FgolUrPvqHLOPRjRynUqBRWTcN5KkRE1Mjs3LkTq1atQlBQkPmRnp4OADhy5AjatWuHPn36IDMzE3feeSc+//xzXLp0ycu1tsSgYpeEYuvVaXnlDxERNTLFxcUYMmQIduzYYfE4dOgQevbsCblcjuzsbCxbtgwZGRn48MMPkZaWhtzcXG9X3YxDP1YkyfRVKwIQKpVUP8G1VIiIrj3+4aY5It6ug5NUKhUMhuqLPzp27Igff/wRycnJUCjsf+RLkoTu3buje/fuePXVV5GUlISFCxdi/PjxUKlUMDo5mdddGFQc4FoqREQEmcypiay+Ijk5GRs3bsSxY8cQFBSEsWPH4vPPP8eIESPw/PPPIzw8HIcPH8b8+fMxa9YsbNmyBb///jv69++P6OhobNy4EefPn0erVq0AAElJSViyZAkOHDiA2NhYhISEeHyuCod+HLAJKpyjQkREPu7ZZ5+FXC5HRkYGoqKiUFFRgT/++AMGgwH9+/dHZmYmxo0bh9DQUMhkMmg0GqxZswaDBg1Cy5Yt8fLLL+Pdd9/FwIEDAZgWfIuLi0NWVhaioqLwxx9/ePw9sUfFAd6YkIiIGpuWLVti/fr1NvsXLFhgt3yrVq2wfPlyh8eLiorCxIkTMWjQIF7142u0NkM/7FEhIiLyNAYVB3hjQiIiIu9jUHGgyPryZM5RISIi8jgGFQds5qjwqh8iIiKPY1BxwPaqHwYVIiIiT2NQsSKE6avNHBX2qBAREXkcg4oD7FEhIiLyPgYVK1VL6Nuuo8LJtERERJ7GoOKAzVU/lWWAQe+dyhAREV2jGFQc0IpA250c/iEiIvIoBhUHbFamBbg6LRERkYcxqDiggxIVQm65k/NUiIiIPIpBxSGJV/4QERF5GYNKLbiWChERkXcxqNSCPSpERETexaBSC66lQkRE5F0MKrWw6VHh0A8REZFHMajUokhYLfrGoR8iIiKPYlCxIkEyf6+F1aJvXEeFiIjIoxhUamGzjD7nqBAREXkUg0otbCfTcuiHiIjIkxhUamGzjD4n0xIREXkUg0otbBZ8Y48KERGRRzGo1IKXJxMREXkXg0otuOAbERGRdzGo1MLmqh9DBaAv905liIiIrkEMKrXQikDbnRz+ISIi8hgGlVrYzFEBOKGWiIjIgxhUrEjVC9NCDwXKhdKyAOepEBEReQyDSh1sr/xhUCEiIvIUBpU6cC0VIiIi72FQqQPXUiEiIvIeBpU62PaocOiHiIjIUxhU6mB7B2X2qBAREXkKg0odbFan5dAPERGRx3g1qEyYMAGSJFk80tPTvVklG1pYLfrGHhUiIiKPUXi7Aq1bt8aKFSvM2wqFd6skhOV2kbAa+mGPChERkcd4PagoFArExsZ6uxoO2Vz1w8m0REREHuP1oHLo0CHExcXBz88PWVlZmDx5Mpo2bWq3rE6ng06nM29rtabeDb1eD71e75L6CGG02LaeoyLKClDponMRzD83V/38yD62s2ewnT2D7ew57mrr+hxPEsJ6sMNzli1bhuLiYqSlpeHMmTOYOHEiTp06hT179iA4ONim/IQJEzBx4kSb/fPmzUNAgJ378jTAR3tlOKStnrrTT7YFn6ummbdLVNFY0fo/LjkXERHRtai0tBT33HMPCgsLodFoai3r1aBiraCgAElJSZg2bRoefPBBm+ft9agkJiYiPz+/zjfqrPtnb8H6oxfN29fJ/sJ81RvmbeEfjsrxB11yLjKl6uzsbPTr1w9KpbLuF1CDsJ09g+3sGWxnz3FXW2u1WkRGRjoVVLw+9FNTaGgoWrZsicOHD9t9Xq1WQ61W2+xXKpUua0Cp5l0JYTv0I5UXQqlQWN69kK6YK3+G5Bjb2TPYzp7BdvYcV7d1fY7lU+uoFBcX48iRI2jSpIm3q2KmtV7wTRgAfal3KkNERHSN8WpQefbZZ7F69WocO3YMf/75J2677TbI5XKMGDHCm9WyoBWBtju5lgoREZFHeHXo5++//8aIESNw4cIFREVF4YYbbsCGDRsQFRXlzWpZKLbuUQEur6XiO70+REREVyuvBpX58+d78/ROMUAOKAMBfUn1Tq6lQkRE5BE+NUfFZ/lZzUjm0A8REZFHMKg4Q20VVHTsUSEiIvIEBhVn+IVYbrNHhYiIyCMYVJxhM/TDHhUiIiJPYFCxYncdN5uhH/aoEBEReQKDijM4mZaIiMgrGFScYT1HhT0qREREHsGg4gzroR/2qBAREXkEg4ozbK764WRaIiIiT2BQcQYn0xIREXkFg4ozuI4KERGRVzCoOMP6qh+uTEtEROQRDCrOsDeZ1mj0Tl2IiIiuIQwqzrDuUYEAKoq9UhUiIqJrCYOKFQl2lqa1nqMCcEItERGRBzCoOEMVDFgHGE6oJSIicjsGFWfIZIA62HIf11IhIiJyOwYVZ3EtFSIiIo9jUHEWb0xIRETkcQwqzrK5MSGHfoiIiNyNQcVZNmupMKgQERG5G4OKszj0Q0RE5HEMKla6NQu3/wQn0xIREXkcg4qVpMhAi+2W0UGmb3hjQiIiIo9jUHGWzY0JGVSIiIjcjUHFip0F9E04mZaIiMjjGFScxaEfIiIij2NQcZbNOioMKkRERO7GoFIHAWH6xmboh0GFiIjI3RhUnGU9mbaiCDAavFMXIiKiawSDirOse1QADv8QERG5GYOKs6znqAAc/iEiInIzBhVnqQIBSW65jz0qREREbsWg4ixJAtTBlvu4lgoREZFbMajUB29MSERE5FEMKvWh5loqREREnsSgYkVyuIY+uDotERGRhzGo1IfN0A/nqBAREbkTg0p9WK+lomNQISIicicGlToIUWODk2mJiIg8ikGlPnhjQiIiIo9iUKkP3piQiIjIoxhU6oOTaYmIiDyKQaU+bCbTskeFiIjInRhU6oPrqBAREXkUg0p9cDItERGRRzGoWJFQy9K01kM/+lLAoHdvhYiIiK5hDCr1YT2ZFuDwDxERkRsxqNSHdY8KwNVpiYiI3IhBpQ41F6aF0h+QKS0LsEeFiIjIbRhU6kOSuJYKERGRBzGo1BfXUiEiIvIYBpX64o0JiYiIPIZBpb64lgoREZHHMKjUl82NCTlHhYiIyF18JqhMmTIFkiRh3Lhx3q5K7biMPhERkcf4RFDZvHkzPvvsM7Rt29bbVambzWRa9qgQERG5i9eDSnFxMUaOHInPP/8cYWFh3q4OpFpW0AfAHhUiIiIPUni7AmPHjsXgwYPRt29fvPHGG7WW1el00Ol05m2t1hQS9Ho99HrX3HOnsrLSYlsIYXFsmTIQ8hrPG8sLYXDRua9FVW3rqp8f2cd29gy2s2ewnT3HXW1dn+N5NajMnz8f27Ztw+bNm50qP3nyZEycONFm/2+//YaAgACX1GnHBQmoEUVKSkqwdOlS83bTC8fQoUb5wnMnsKbG89Qw2dnZ3q7CNYHt7BlsZ89gO3uOq9u6tLTU6bJeCyonT57EU089hezsbPj5+Tn1mhdffBHjx483b2u1WiQmJqJ///7QaOzch6cBpD1nMfvgLvN2YGAgBg26ofr5/QbgxBfm7VA/CYMGDXLJua9Fer0e2dnZ6NevH5RKZd0voAZhO3sG29kz2M6e4662rhoRcYbXgsrWrVuRl5eHjh07mvcZDAasWbMGH330EXQ6HeRyucVr1Go11Gq1zbGUSqXLGlChsG4SyfLYgeGWz+qK+IfiAq78GZJjbGfPYDt7BtvZc1zd1vU5lteCSp8+fbB7926LfWPGjEF6ejpeeOEFm5DiM2zWUeFkWiIiInfxWlAJDg5GmzZtLPYFBgYiIiLCZr9Psb7qx6AD9OWA0rnhKyIiInKe1y9PbnSse1QALqNPRETkJl6/PLmmnJwcb1ehbtY3JQRMwz9B0Z6vCxER0VWOPSr1pVADCqthHq5OS0RE5BYMKg3BGxMSERF5BIOKlbpW0AdgO/zDK3+IiIjcgkGlTsJ2l82NCRlUiIiI3IFBpSF4Y0IiIiKPYFBpCOuhH/aoEBERuQWDSkNwMi0REZFHMKg0BId+iIiIPIJBpSGsgwqHfoiIiNyCQaUhOPRDRETkEQwqDWGzjgqDChERkTswqDQE11EhIiLyCAaVhuDKtERERB7BoGJFslpDX9hZmNbuZFq7BYmIiOhKMKg0hPXQj7ES0Jd6py5ERERXMQaVhrDuUQE4/ENEROQGDCoNoQ623ccJtURERC7XoKDy1Vdf4ZdffjFvP//88wgNDcX111+P48ePu6xyPkuuBJSBlvvYo0JERORyDQoqb731Fvz9/QEA69evx8cff4ypU6ciMjISTz/9tEsr6LO4lgoREZHbKRryopMnTyI1NRUAsGjRItxxxx145JFH0L17d/Tu3duV9fNdag1QdKZ6W8egQkRE5GoN6lEJCgrChQsXAAC//fYb+vXrBwDw8/NDWVmZ62rny7iWChERkds1qEelX79+eOihh9ChQwccPHgQgwYNAgDs3bsXycnJrqyf7+KNCYmIiNyuQT0qH3/8MbKysnD+/Hn8+OOPiIiIAABs3boVI0aMcGkFfZbNjQkZVIiIiFytQT0qoaGh+Oijj2z2T5w48Yor5GtyLzhYyI2TaYmIiNyuQT0qy5cvx7p168zbH3/8Mdq3b4977rkHly5dclnlvGH3KScDB29MSERE5HYNCirPPfcctFrTB/Pu3bvxzDPPYNCgQcjNzcX48eNdWkFPW7DtlHMFreeocOiHiIjI5Ro09JObm4uMjAwAwI8//ohbbrkFb731FrZt22aeWNtYyazvSugIJ9MSERG5XYN6VFQqFUpLTXM3VqxYgf79+wMAwsPDzT0tjZWzOcV2Mi3nqBAREblag3pUbrjhBowfPx7du3fHpk2b8O233wIADh48iISEBJdW0NOc71HhVT9ERETu1qAelY8++ggKhQI//PADZsyYgfj4eADAsmXLcPPNN7u0gp4ma2iPClemJSIicrkG9ag0bdoUP//8s83+995774or5G1Sg+eoFAFGIyDjDamJiIhcpUFBBQAMBgMWLVqEffv2AQBat26NW2+9FXK53GWV8wan56hYD/0II1BRbLufiIiIGqxBQeXw4cMYNGgQTp06hbS0NADA5MmTkZiYiF9++QUpKSkuraQnyZ1NKtZDP4Dpyh8GFSIiIpdp0DjFk08+iZSUFJw8eRLbtm3Dtm3bcOLECTRr1gxPPvmkq+voUZ2SwpwrqA623ccJtURERC7VoB6V1atXY8OGDQgPDzfvi4iIwJQpU9C9e3eXVc4bhndJxPzNJ+suKJObelVqrp/CtVSIiIhcqkE9Kmq1GkVFRTb7i4uLoVKprrhS3nS2sNxmX0Fphf3CXEuFiIjIrRoUVG655RY88sgj2LhxI4QQEEJgw4YNeOyxx3Drrbe6uo4eZW+Kyqbci/YLcy0VIiIit2pQUPnggw+QkpKCrKws+Pn5wc/PD9dffz1SU1Px/vvvu7iKnhUX6m+zb9baXPuFuZYKERGRWzVojkpoaCgWL16Mw4cPmy9PbtWqFVJTU11aOW+IClbb7DtXZDscBIA3JiQiInIzp4NKXXdFXrVqlfn7adOmNbxGXhar8bPZ989eDi63th764WRaIiIil3I6qGzfvt2pck6v7Oqj7NU/yM9BM3EyLRERkVs5HVRq9phc7dJjg7H/bPVVTQajsF+Qk2mJiIjcijemsUNudWfCSoODoGIzmZZBhYiIyJUYVOxQWAUVg3DUo8LJtERERO7EoGKHzDqoOBz6sQ4qnKNCRETkSgwqdhSVV1ps6w1G+wU59ENERORWDCp2HM4rttj+43C+/YKcTEtERORWDCpO+HXvOftPWA/9VBQBRoP7K0RERHSNYFBxgvVVQGbWQz8Ah3+IiIhciEHFCTe3ibX/hPXQD8DhHyIiIhdiUHHCL7vO2H9CFQRIVk3IHhUiIiKXYVC5EpJkZxl9BhUiIiJXYVCxo22C5STZm9KjHRe2ufKHa6kQERG5CoOKHV2Swy22I4NUjgurra784dAPERGRy3g1qMyYMQNt27aFRqOBRqNBVlYWli1b5s0qAQBUCstm0ZZVOigJrqVCRETkRl4NKgkJCZgyZQq2bt2KLVu24KabbsLQoUOxd+9eb1YLx/JLLLaX7z3ruLD1Wio6Dv0QERG5isKbJx8yZIjF9ptvvokZM2Zgw4YNaN26tZdqBSzbU0swscbJtERERG7jM3NUDAYD5s+fj5KSEmRlZXm1Lglh/s4X5mRaIiIit/FqjwoA7N69G1lZWSgvL0dQUBAWLlyIjIwMu2V1Oh10Op15W6s19V7o9Xro9XqX1emergmY+ushi32Oji9TBkFeY9tYVgCDC+tytatqV1f+/MgW29kz2M6ewXb2HHe1dX2OJwkhhEvPXk8VFRU4ceIECgsL8cMPP2DWrFlYvXq13bAyYcIETJw40Wb/vHnzEBAQ4LI6bciT8L8j1fGjaaDAM23t38Mn9dwvaH36W/P2ueBMbEh9zmV1ISIiutqUlpbinnvuQWFhITQaO6u81+D1oGKtb9++SElJwWeffWbznL0elcTEROTn59f5Rutj0Y7TeO7HPebtVrHBWDLW/nCUtP1rKJaON28b4zvDMHq5y+pytdPr9cjOzka/fv2gVCq9XZ2rFtvZM9jOnsF29hx3tbVWq0VkZKRTQcXrQz/WjEajRRipSa1WQ61W2+xXKpUubUC1yvJYBiEcHz8gzGJTptNCxj+cenP1z5DsYzt7BtvZM9jOnuPqtq7PsbwaVF588UUMHDgQTZs2RVFREebNm4ecnBz8+uuv3qwWVHLLuyXrDbV0OnEdFSIiIrfxalDJy8vD/fffjzNnziAkJARt27bFr7/+in79+nmzWlDKLS+G0huMjgtzZVoiIiK38WpQ+eKLL7x5eocU9Qkq1gu+6UsBgx6QszuSiIjoSvnMOiq+RHklQz8Ah3+IiIhchEHFjvoN/dgJKlxGn4iIyCUYVOyoV1BR+gMyqxE09qgQERG5BIOKHQqZ5dBPZW1DP5Jk58aEDCpERESuwKBih0ph2SyVRoFa18XjjQmJiIjcgkHFDuuhH6C+a6lwjgoREZErMKjYYT30A9RzQi2HfoiIiFyCQcUO66EfoJ5rqXDoh4iIyCUYVOyw36NS29APJ9MSERG5A4OKHcr69qjYTKYtcG2FiIiIrlEMKnYoZbbNUuslyrwxIRERkVswqNhhvYQ+AFRwMi0REZHHMajYIZdJkKyySqWRk2mJiIg8jUHFDkmSbCbU6iu5jgoREZGnMag4oLJa9I1DP0RERJ7HoOKAQm59v5/ahn44mZaIiMgdGFQcsL2Dcm1DP6GW2wYdoC93faWIiIiuMQwqDtjMUaltMq310A/A4R8iIiIXYFBxwKZHpbIeQz8Ah3+IiIhcgEHFAeugUmmsZehHoQbkast9Ol75Q0REdKUYVBzQVRostitq61EBuJYKERGRGzCoOHCqwHIy7Nbjl2p/gfXwD+eoEBERXTEGFSfN3XC89gI2Nybk0A8REdGVYlBxFa6lQkRE5HIMKq7C1WmJiIhcjkHFVTiZloiIyOUYVByI1VhebjysfVztL7AJKpyjQkREdKUYVBzo0SLSYjssUFX7Czj0Q0RE5HIMKg7sPW0ZNGb/caz2F9hMpmWPChER0ZViUHHgrzNF9XuB9dAPe1SIiIiuGIOKq3AdFSIiIpdjUHEVrqNCRETkcgwqrmJvMq2o5UaGREREVCcGFVexnqNirAT0Zd6pCxER0VWCQcWBW9s2qd8LrIMKwHkqREREV4hBxYF+GdE2+w6eq+VKIHWw7T5e+UNERHRFGFQcCFTLbfZtyr3o+AVyJaAMsNzHCbVERERXhEHFgZbRQTb7Fmz7u/YX2aylwqEfIiKiK8Gg4kCMxs9mX2SQ2k7JGmzWUmGPChER0ZVgUKmHYl1l7QW4jD4REZFLMajUw59HLtRegDcmJCIicikGFVfi6rREREQuxaDiSrwxIRERkUsxqLgSb0xIRETkUgwqrsShHyIiIpdiUHElNYd+iIiIXIlBxZWs56iwR4WIiOiKMKjUk9EoHD/JdVSIiIhcikGlFv5y21CiNxodv8BmHRUGFSIioivBoFKLW5rahpItxy45foF1j4quCBC19MAQERFRrRhUahHlb7tv5KyNjl9gPUdFGIGKYtdWioiI6BrCoFILhVTP3hDroR+A81SIiIiuAINKLeID6/kCdbDtPl75Q0RE1GAMKrXwk9fzBTI5oLIKK1xLhYiIqMEYVFyNq9MSERG5jFeDyuTJk9GlSxcEBwcjOjoaw4YNw4EDB7xZpSvHGxMSERG5jFeDyurVqzF27Fhs2LAB2dnZ0Ov16N+/P0pKSrxZrStjc2PCAq9Ug4iI6Gqg8ObJly9fbrE9Z84cREdHY+vWrejZs6eXalU3IQQkSbL/JId+iIiIXMarQcVaYaHpUt7w8HC7z+t0Ouh0OvO2VmsKAXq9Hnq93qV1qe14ZwtKEBmktvucXBVk0U1lKCuA0cV1u5pUtbOrf35kie3sGWxnz2A7e4672ro+x5OE8I2lU41GI2699VYUFBRg3bp1dstMmDABEydOtNk/b948BAQEuKVeT623zXKvdqhEhJ/98m1PzkGz/JXm7dzIm7ArcbRb6kZERNQYlZaW4p577kFhYSE0GjtrkNXgM0Hln//8J5YtW4Z169YhISHBbhl7PSqJiYnIz8+v843Wl16vR3Z2tt2gcleneLw5rLXd18lWTYL8z+nmbWPr22EYNtOldbuaVLVzv379oFQqvV2dqxbb2TPYzp7BdvYcd7W1VqtFZGSkU0HFJ4Z+nnjiCfz8889Ys2aNw5ACAGq1Gmq17ZCLUql02y9rt2Zh2JhreX+f77aewtQ729t/gX+oxaasohgy/iHVyZ0/Q6rGdvYMtrNnsJ09x9VtXZ9jefWqHyEEnnjiCSxcuBArV65Es2bNvFkdux6+Ibl+L+BkWiIiIpfxao/K2LFjMW/ePCxevBjBwcE4e/YsACAkJAT+/nbuCOgFPVIj6/cCv1DLba6jQkRE1GBe7VGZMWMGCgsL0bt3bzRp0sT8+Pbbb71ZLQsymYPLkB2xWUeFNyUkIiJqKK/2qPjIPF7X4tAPERGRy/BePw20/6yDAGLdo1JRBBgN7q8QERHRVYhBpYGemLfd/hPW9/oBAF2ReytDRER0lWJQaaDDecX2n7Ae+gE4T4WIiKiBGFSc8PLgVs4XVgUBklWz8sofIiKiBmFQccKdnROdLyxJgDrYch8n1BIRETUIg4oTVHL7zeTwqiW11TwV9qgQERE1CIOKE/xVcrv71xzKt/8C6wm17FEhIiJqEAaVKzAj57D9J2zWUuFkWiIiooZgULkCG45etP+E9VoqOgYVIiKihmBQcQeuTktEROQSDCpOahkTZHf/xZIK253Wc1Q4mZaIiKhBGFSc9H83tbC7v+OkbNudvDEhERGRSzCoOGlgm1jnC3Poh4iIyCUYVJykcLCWil02k2kZVIiIiBqCQcUFisr1lju4jgoREZFLMKjUw8E3BtrdnznhN8sdXEeFiBq7ygrg763AhhnADw8A/70DWDEROLMTcLQqN5EbKLxdgcZEpXAy13EJfSJqbLRngL83ASc3AX9vBk7vAAw6yzKHVwDrpgHhzYHWtwEZw4DYTNM9zhobgx64mAsUngCMBkAYTQFMGE0PCMt9QO1lJAkISQSi0oDgJo2zTXwUg4qLFJbpEeKvNG1Y96joS01/FHKl5ytGRGStsgI4u+tyKNkE/L0FKDzp/OsvHgXWvmt6hKeYQkvr24CY1r73AV1ZYarv+X3A+QPA+f1A3n7gwmHAqK/79Q2h1gCRLYDINCCq5eWvaUBoEiD3wMduRQlQnAeUnAcCIoCIFPef040YVFyk3cTfcGzKYNOG9WRaANAVAQHhnq0UEREAaE9X95Q46i1pqItHgLX/MT0iUqtDS3SGZ0NLpQ64cMQ2kFw8AhgrPVcPwNSLfmqr6VGTXGVqo8iWpuBS9TUiFVD6Oz6eEKbPkJLzlwNIXnUQsfiaBxSfB/Ql1a+9/kmg/yT3vE8PYVCpp9HXJ2POn8fsPne2sByxIX62k2kB0zwVBhUi32TQA6UXAH0ZUFle/bWyHNCXA5VlDr5alpFXlKLLuTzIVm4GYjIufxilAWr7C0a6XKXO1FOQt8/0OL/fFEq0fzf8mJEtgYSupn/X9v8EFJxwXPbCYWDNO6ZHRIsaoaXVlYcWIUz/jpbkQ9KeQfzF9ZDl7AQuHrwcSI4CwnBl53A3QwWQ95fpYUECQpuafl/CU0xBo/h8dfAoyTP9njVEyfkrrra3MajU0+M3pjgMKv/8ZisWPt7dlIxlCssUzwm1RN5Xcxggb7/pg/x81TDAlf+vWwYgDgDWW/1PumruQlR69dfIloB/aMNOZNBXB5Lz+6u/XjhyZR/Wag0Q3wlI7GoKJ/EdLf+DNeBN4PQ2YO9CYO9i0/wORy4cAtZMNT0i04DWw6pDSxV9OVCab/owLcmv7h2o2rb+/vJQjQJAZwA43vC3ahYQCagCAEiAJDMFKklWyzbsP19ZDlw4atmb4TQBFBw3PVytOM/1x/QwBpV6ig72c/jc9hMFpm8kyfQHX1bjpoWcUEvXKoMeKDoDFJ4CtJcfhadMvRIBkUBg1OVHje8DIq5sLL+ywtTln1c1DHD5q4sCSb0VnjQ9Dq+w3B/c5HJwaWUZZKrCgaHSKljtc+38isg0ILGLKZQkdDGdWyZ3XF6STEEmvhPQbxJwahuwdwHw1+La57jkHwBWv216hDUzHack37P/LtbV1q5gNJp+v/MPAOcPWn4tveC68zhL4Vf7z7ORYFBpgLkPdsV9X2yy+1yXN1dg87/7mibU1gwqXEuFrkZGA1B87nII+bs6jBT+bZoXoT0FFJ0F0IDLWf3Da4QXR4Em0hQ8qoJIVTDxxryEhig6Y3oczbHcHxht+gC9eNQ0XOAK6hAgoVN1KEnoBPiHNfx4knT5eJ2A/m+Y5mPsXQjsXVT7UNOl3Iaf0xma+Msh5HIQiW51Zb1X9SGTAaGJpkdqX8vnSi5cDi4HgPyD1V/rM4kZAJSBQFCU6XckKNr0d2DxtcZ+dbDvTW5uAAaVBuieEunwufNFOlwqqUAYb0xIjV1VCKkKHFVfC2v0ihSdcd+8gLKLpkf+QfccvzZytel/o0q/y1/9nfpqkKlwZP8upIYYIMs/aPpQrrq0tT5KLk+MbChNAhB9+cM6OsM0hBOZZvogdQdJAhI6mx79JlWHlr8WmX5XXEyoAlGKQPgntoXM/D4vBxLrqy59RWAEEHg9kHS95X5dsWmYrKrnpfBvU8AIjLYfSFSB3qm/FzGoNIBMVntC7TApG7npGliU4hwV8iWVFaaQoT0NFJ2+HEJqBpLTpp4QX5+c2FDBTar/1131QReeYvoQUPg1+APdqNdjn3Ypmg0aBJlSaZqDceHw5bkwB6q/uqrHp+p9RLcyPaqGNbz5YS2TmYaTErtc7mnZcjm0LHYcWmQKOz1nDnrSAiJRKSmxYulSDKpq58ZMHQTEdTA9yC4GlQY6+MZAtHx5mcPnz+vViK65g0M/vq2ywtQFW3B58Sd1sO2jMYz16surL10szjP1iBTnQVZ0Fl2P7oD8i3dNAeVK/rfeUAo/U7d8SLzpf/yqgMuTJPOrJ1SWXkSDhokcCY6rDiI1hwM8MQwAmHpkYtuYHjWZJ/XWDDD7gfxD9ueeBEZffh+taoSStCsbuvEEmcw0MTexK9D/TVNPS/6Byz0GNQKIX2j9hij0blr/hHwSg0oD1bVK7ZoTFfhHzc81Dv14l9EIFJ8FLl2eWW/9VXuq7i56ZaD9AKPW2N+n9Dct8idTmNZPkCsBmdI0SVSmrPGc0nZbrqoORgb95XUSzlkGEJt9eYDOfs+dHEATAHBXx55cBWjiTAEkJP7y9/FASEL1V/+wuj+MDJWm4R6bKz+sv17+vqLI9Lqa8xKiawQSe0sF+AKFylTP6HTL/YZK03BR3j5TL2x4c1MouRqWNqjZ00JUDwwqV2DLy33R+Y0Vdp/TigCLbVFeiMY/pcmHCQGUXbIfQi4dN/WUXOkCV/qSy+sbnHVNneskmUKLqyZTNpRcZRpi0FQFkLjqAFL1fUCka+Y/yBWmcfig6LrLAqb1ToTx6hm3lysur2jawts1IfIZDCpXIDJI7fC5IlgGlQv5eXA8BfcaJ4TpNgPl5aaep3Kt6av5+yLn9utLvf1OXEy4P6Qo/KvDR80goqnRKxIQ4b5JmFeqttU8ieiqwKByhRaP7Y6hH/9hs18rLP8BDfl7FfDJ9aZZ3EEx1TO4zd/HmLYDIhrHXIj6MBpMQyuXjgOXjpkeBabvFZeOY0jJBch2+MikTXWIae6Errh6WKExkquqrxYIioYxIBKHzmiR0rEXFGFNqwOJM8MxRERexKByhdolhuK1IRmY+JPlksjWPSpKUQHk7QXqmsMoyUzd6FXd3zU+bCBXma4UMFaa5i0YDdXbxko729aPy88rLi/zb/EItdz2v7yt8Kv7g6xq2KVGALEIJYV/O1ycSrr88BiFn2mp6tAkICzJ9mvNyYlGI1BRfLnnpqi6B8diu6i6l6fmvnKtaaVK889Kb/pa9X1DrviQ5DUuU4y5/IiyCruX91lNTjTo9di/dCmadxwENParJIjomsKg4gJjujezCSpHjU0adjBhrF5D4ZwLKnel5Co7oSakeuXdS8eASyccTuL0OElmmtBZM3yEJVd/Hxjt/DCGTGa6zNMdl3oK4TjEGPTVzxkqTOEqKMYUonx1CIaIyE0YVFzE+nLlraIl/ld5I+6S50AuufByS08zVFRfZeEpCj9TEPLTVF9B46e5fHWNvf3BpiGbwEjTxE55I+gxkC5PlG0MdSUi8iIGFRdRKWR4qk8LTP/9EABAQIYXKx/G5MoRiJMuIkoqQCQKTV8lrXm7a1QllOX5kJVegEvXj/BFqqDLPRvJlx9JqAxOwLqdh9G9z0AoAyNMoUOh8nZNiYjIRzCouNDT/Vrif5tOIK+o+jJYLYKgFUHYL5raf9Fp0xc5DFg4Og1tQ3SX7yBatVjX+eoFvISxxrobCtOkW5mixqNq29HzCtPQgb4cKC8wrdNg71FR3LAGkOSmNTSqgog5lDQzDbsERNjMdxF6PQoPVZrKcO4EERFZYVBxsU3/7ouV+8/hgTlb6vU6A+S4dc5hdE4Kw9R/XIfmLYLcVENnKqO/PCG0ACgrcBxo1EGWoaSxDLsQEVGjwaDiBjelxzi8bLkuW45fwk3vrgYAjOiaiJSoINyflVznSrguJVeaVsK8GlbDJCKiRo1BxU3aJYbivw92w71fbGzwMf63yXT77zd+2QcA+GfvFPTLiEH7hNA6b4xIRER0NWBQcaMbWkRi/Ys3IWvySpccb0bOEczIOWKx7/mb0/DPXimQuGgXERFdhRhU3KxJiD9yJw/CkfPF6DttjcuPP3X5AUxdfsBi3+TbM3Fd8whEBqkQ7Mc5I0RE1HgxqHiAJElIjQ7Grgn90XbCb24/34sLdpu/v71jPO67LgnntOVonxiG2BA/t5+fiIjIVRhUPEjjp8SxKYNxtrAc103+3SPnXLDtFBZsO2Wzv2l4ANonhuKpvi2QX6RDqzgNNOx9ISIiH8Og4gWxIX44/OZA5BXpMOzjPyzWXfGUExdLceJiKZbsPG3z3M//dwMig9QI8VfCXyWHEIJzYIiIyCsYVLxEIZchLtQfm/7dFwCgNxgxf/NJvLJoj5drBtzy4bpanw8NUOKzezuhWWQg/i4og8ZPgdToYA/VjoiIriUMKj5CKZfhvuuScN91SQCASoMR72YftLnKxxcUlOoxfOYGu8/d3SURreM0eGXxXoQGKPHGsDZIjgiEQi7h05wj0BsEnu7XEqnRXlzQjoiIGg0GFR+lkMvwws3peOHmdBzOK8ZHKw9h0Q7bYRpfM3/zSfP3BaV6PDFvu02ZX3afwbD2cchMCMVXf+bCUC5HYOp53NAyBn5KuSerS0REPo5BpRFIjQ7C+3d3wPt3d0BFpRFlFQbsPlV4RYvJeduiHadrBC8JD821DDTxof7o2yoaAzObIEAlx6bci/BXyREVpEa3ZhEICVCiqFwPpVwGP6UcRqPAntOFCAtQITE8wPNviIiI3IJBpZFRKWRQKWS4oUUkjk0ZbN6vNxix74wWt35U/2X7fdGpgjJ8tf44vlp/vEGv/3J0Z0QGqTFrbS4ig9R4sk8q/FVyqBXssSEiakwYVK4SSrkMbRNCcWzKYAgh8EnOEazcn4eHbmiGPacL8fEq35vr4k7WN4X88o/cWsv3aBGJx3unok28BgCw6+9CzN98EmUVBoy9MQUxGj/EhfoDMM0fKqkwIEitgJy3MiAicisGlauQJEkYe2Mqxt6YCgAYmNkEzw1IBwAYjAJfrz+GiT/95c0q+py1h/Kx9lC+3edW7Dvn1DFSogLx1m2ZUCvlaBEdhHPacgSoFPj7UiliNH61Dkmd05Zj+4kCtI7TcOiKiKgGBpVrjFwmYUz3ZhjTvZl5nxAC57Q6VBqNKCjVAwDOF+nw5PztKCqv9FZVG50j50scXg1lbeZ9ndChaRgOnC3C52uPYvXB8xbPz3uoG/KKdPgk5zBiNH5467ZMJIYHoFhXCZ3egIggtc0xDUYBXaUBfgo5b1pJRFcNBhWCJEnmpfUTwqr3754wwPz9qYIy6PQGNI8KQkFpBSb+9BcWbrdd8Zac88jcrbU+f8+s6onSB88Vo8fUVbWWDw1QmkMmALRNCME/e6VgYGYTLNl5Gl+uO4rT5+UojDqJ+LBAJEcGIiUqCKUVlThdUIb40AD4q658/o7BKCCEgEIuu+JjEREBDCrkpPjL8zMAIDRAhfeGt8d7w9vblNOW67H12CVsP3EJH6w87MEaXttqhhTANMfmn99ssyol4dUl++p97HkPd0Nhqd58vDbxGhSU6lFRacSwDvF4fkAa5DIJv+w+Y74cvVNSGMbemAIA6N0y2m4Pj9EoUGEwQiWXXXEPULneAJkkQaVgQCK62jCokEtp/JS4MT0aN6ZHY3z/NIvnrP+3vfNkAUbP3oRLpXo0CfHDmcJyb1SZ6nDP55aXwe85pTV/P3PNUcxcc9TmNVuPX7KZ0FzFXylHmd5gsS8+1B+nCspsyk4a2ho9W0ZBgoR/fPon8op0iNX44am+LXCmoAxDO8Tjp52n8f6KQ+bjFJXr0SwyEFPuaIsSXSUSwwMQo+HNOIkaKwYV8hjTFTLV/3NulxiKTS/eiKVLl2LQoJ6QyxUorqhEgFIOhVyG7ScuoVxvREp0IFSX10vJ0+rwyNwt2H+2yHtvhK6IdUgBYDekAMAri/fa7DurLTffIdy6167qODv/LsTA6WsbVL8+6dH4fX+eeXvLy31xvkiH7L/OoV9GDNQKGZbtOYukiADcmBaNr9cfx3dbTuKOjvHo3jwcOWcknFyTi97pMWgTHwIAOHSuCHd9th6t40Iwe0wXKK2Gxvaf1SJPq8N1zSOgUshgNApIEizusSWEwNLdZ3Eorwi3totD8yiu7kzXBq8GlTVr1uCdd97B1q1bcebMGSxcuBDDhg3zZpXIi2QyyeIOzh2ahtmUaRoRgOXjetrsNxoFjDV6a4p1lVi1Pw/fbTkJg1EgRuPHOTXklJohBQA6v7HC/P207IMOX/ef3w7iPwAAOXDsEP6TfcimzLrD+Wjx72VXXMeqHiQAeLx3ChQyyRza3r4jE/nFFbghNRLtEkPrPFZhqR5qpcxmVWghBA7nFeOHrX/DYBS4PysZTSMCbMoApt5SAD47N0kIAaMAZFbhjxoHrwaVkpIStGvXDg888ABuv/12b1aFGjmZTIKsRm9NkFqBIe3iMKRdnHmfvTk1ZwrLkKfVIS7UH0fPF2PMnM0ovbx2ihDApdIKzN98Epf/PSbyOZ9Y3Q/shR9NvU3v/HrApeeZta72tYgcGdE1ES8NaoW8Ih2O5BVj8Y7T+GX3GQCm+U/7zhTh933nsP9sES6WVKB7agQmDGmN2X8eQ4muEmNvTEVSRADm/HEMfx65gJSoILSND8beixIO/X4YN7aKQaekcLvnPl+kw1d/HsNHq0whrl9GDD64u4PFxHEhBMr1RijkEr5efxybci/gnm5J6NUyqkHvl1xPEsI3/gmWJKnePSparRYhISEoLCyERqNxaX30ev3lIYlBUCqVdb+AGqSxtLOu0gCDUSBAZcr2h/OKUayrRLuEEJwqKMPZwnIUlulxsaQCQgD7zxZh4fa/oVLI0DQ8AB2TwvDzzjMOhziIyPcMyozF0t1nHT5/Q2ok+rSKxg2pkTAK4N4vNuJ8kQ7BagWeHZCGiT/txeXOJtyflYS7OieiwmDEn4fz0aNFFJIjAnG+uBx9p60xH3NY+zjc3jEBcaF+iNb44b8bjmPqclPofLpvSzzVtwVOXiw1/7uTGB6AtgkhEAL4aedp/HfjcSSEBeD1W1tD46+84l4kd/0bXZ/P70YVVHQ6HXQ6nXlbq9UiMTER+fn5bgkq2dnZ6Nevn09/gDZ213o7a8v0kCQJwX4KFJbpcb5Ih+SIACjkMpzTlmP6yiP4cdsp8z92RESe1jnSiMcGdECv9BiXHVOr1SIyMvLqCyoTJkzAxIkTbfbPmzcPAQFczZOuPZVGYM8lCeUGoMIARPsDUX4CoWpALgEGAZRXAufKAJUcCFIAFUbgRLGEQ1oJG/J8c04BEfme6VmuWwC0tLQU99xzz9UXVNijcvVhO3uGK9tZbzBCJkk4cbEUuRdK0SUpDBUGI/acKkTT8ABEB6uxYPtp7DhZCF2lAQNax8BoFHj2xz0uejdE5A2HJvV32bHq06PSqC5PVqvVUKttlw5XKpVu+5Bz57GpGtvZM1zRzlUvb9lEhZZNQs37Y0MDzd8/0CPF5nX/6JIEwHSFFoB6L/JmNAqU6Q0QMK3FUq43oERXiUC1Ag/M2YyNuRctyreMCUJekQ5CAIVlevsHJSKnufLf6Pocq1EFFSJq/Bq6Cq1MJiFQXf1PVqBaYd6e/8h1+OuMFhGBavPtIBoiv1iH/GIdWkQHW9wZu/zy2i/HL5TinLYcMklCy5gglFYYsHTPGRw8W4Rdpwpx9HwJAOCfvZphRNdkZO87B12lAS2ig/HH4XzM23QCFZXGBteP6Frk1aBSXFyMw4erF2zKzc3Fjh07EB4ejqZNm3qxZkTUmEiShNZxIVd8nMggNSLt3PCxao2RtNhgpMUGWzz3eO9U8/fmKyT6toBSqcSDN1Tf/LNfRgwm3Nq6wXUTQkAIy6CnqzSgoFSPiECVeQ2TQ+eKsP7oBWTGhyAjztSlPuePY/hiXS46NA3Ffdcl43BeEXSVRsRo/PDbX2dxptB0924AuPe6pkgKD8Q5bTmW7TnLK9XI67waVLZs2YIbb7zRvD1+/HgAwKhRozBnzhwv1YqIyPdIkgTrq0zVCjliNJYLtbWICUaLGMsw9WivFDzaq3o47oYWkebvh3WId3jOl2/JwPkiHYp1lUiOCHDqMldH93A6W1iO3acK0TrOdK8of5UczSKrhwv1BiMMRmEOhQajwLebT+JwXjEe6dkcFZVGVBiMCA1Q4o+Defhg2U4UGNW4tX0cnu2fBpVChpX787ByXx5uzoxFeIAKJy6W4uj5Ery3wnKhvsggFfKLKxAVrMb5Ih2obm/f3vCQfaW8GlR69+4NH5nLS0REdkQFqxEVbNvL5IhMJsFPZnsn7tgQP/OwXFyNm5xWUcplqLk4rlwm4Z5u9nvWB2XGAie3YdCg3hZzHQa0jsWA1rHm7aqVeZ/q28Lp+jvLaBTQlutRaRT443A+YjR+uK55BABTr9apgjKkRgchVuOHJTtP49SlMtzeKcHiBq/acj32nylCWkww1EoZluw4DYVcglohx7ELJQjxV+K/G46bbxlye8d4RAWrUVCix7dbTlrUp0eLSKw9lO/y9wkAXaOMuL2WQOtunKNCRERUTzKZhNAAFQBgaHvLD3HrXq3bOybYPYbGT4muzapX1b2rS6JNmXuvS7L72rf/0bbedW6IquFMb+IiCkREROSzGFSIiIjIZzGoEBERkc9iUCEiIiKfxaBCREREPotBhYiIiHwWgwoRERH5LAYVIiIi8lkMKkREROSzGFSIiIjIZzGoEBERkc9iUCEiIiKfxaBCREREPqtR3z1ZCAEA0Gq1Lj+2Xq9HaWkptFqtxW3EybXYzp7BdvYMtrNnsJ09x11tXfW5XfU5XptGHVSKiooAAImJtrfGJiIiIt9WVFSEkJCQWstIwpk446OMRiNOnz6N4OBgSJLk0mNrtVokJibi5MmT0Gg0Lj02VWM7ewbb2TPYzp7BdvYcd7W1EAJFRUWIi4uDTFb7LJRG3aMik8mQkJDg1nNoNBr+IXgA29kz2M6ewXb2DLaz57ijrevqSanCybRERETksxhUiIiIyGcxqDigVqvx2muvQa1We7sqVzW2s2ewnT2D7ewZbGfP8YW2btSTaYmIiOjqxh4VIiIi8lkMKkREROSzGFSIiIjIZzGoEBERkc9iULHj448/RnJyMvz8/NCtWzds2rTJ21XyKWvWrMGQIUMQFxcHSZKwaNEii+eFEHj11VfRpEkT+Pv7o2/fvjh06JBFmYsXL2LkyJHQaDQIDQ3Fgw8+iOLiYosyu3btQo8ePeDn54fExERMnTrVpi7ff/890tPT4efnh8zMTCxdutTl79cbJk+ejC5duiA4OBjR0dEYNmwYDhw4YFGmvLwcY8eORUREBIKCgnDHHXfg3LlzFmVOnDiBwYMHIyAgANHR0XjuuedQWVlpUSYnJwcdO3aEWq1Gamoq5syZY1Ofq/VvYsaMGWjbtq15MausrCwsW7bM/Dzb2D2mTJkCSZIwbtw48z62tWtMmDABkiRZPNLT083PN8p2FmRh/vz5QqVSiS+//FLs3btXPPzwwyI0NFScO3fO21XzGUuXLhX//ve/xYIFCwQAsXDhQovnp0yZIkJCQsSiRYvEzp07xa233iqaNWsmysrKzGVuvvlm0a5dO7Fhwwaxdu1akZqaKkaMGGF+vrCwUMTExIiRI0eKPXv2iP/973/C399ffPbZZ+Yyf/zxh5DL5WLq1Knir7/+Ei+//LJQKpVi9+7dbm8DdxswYICYPXu22LNnj9ixY4cYNGiQaNq0qSguLjaXeeyxx0RiYqL4/fffxZYtW8R1110nrr/+evPzlZWVok2bNqJv375i+/btYunSpSIyMlK8+OKL5jJHjx4VAQEBYvz48eKvv/4SH374oZDL5WL58uXmMlfz38SSJUvEL7/8Ig4ePCgOHDggXnrpJaFUKsWePXuEEGxjd9i0aZNITk4Wbdu2FU899ZR5P9vaNV577TXRunVrcebMGfPj/Pnz5ucbYzszqFjp2rWrGDt2rHnbYDCIuLg4MXnyZC/WyndZBxWj0ShiY2PFO++8Y95XUFAg1Gq1+N///ieEEOKvv/4SAMTmzZvNZZYtWyYkSRKnTp0SQgjxySefiLCwMKHT6cxlXnjhBZGWlmbevuuuu8TgwYMt6tOtWzfx6KOPuvQ9+oK8vDwBQKxevVoIYWpTpVIpvv/+e3OZffv2CQBi/fr1QghToJTJZOLs2bPmMjNmzBAajcbcrs8//7xo3bq1xbmGDx8uBgwYYN6+1v4mwsLCxKxZs9jGblBUVCRatGghsrOzRa9evcxBhW3tOq+99ppo166d3ecaaztz6KeGiooKbN26FX379jXvk8lk6Nu3L9avX+/FmjUeubm5OHv2rEUbhoSEoFu3buY2XL9+PUJDQ9G5c2dzmb59+0Imk2Hjxo3mMj179oRKpTKXGTBgAA4cOIBLly6Zy9Q8T1WZq/FnVVhYCAAIDw8HAGzduhV6vd7i/aenp6Np06YW7ZyZmYmYmBhzmQEDBkCr1WLv3r3mMrW14bX0N2EwGDB//nyUlJQgKyuLbewGY8eOxeDBg23ag23tWocOHUJcXByaN2+OkSNH4sSJEwAabzszqNSQn58Pg8Fg8QMCgJiYGJw9e9ZLtWpcqtqptjY8e/YsoqOjLZ5XKBQIDw+3KGPvGDXP4ajM1fazMhqNGDduHLp37442bdoAML13lUqF0NBQi7LW7dzQNtRqtSgrK7sm/iZ2796NoKAgqNVqPPbYY1i4cCEyMjLYxi42f/58bNu2DZMnT7Z5jm3tOt26dcOcOXOwfPlyzJgxA7m5uejRoweKiooabTs36rsnE10Lxo4diz179mDdunXerspVKS0tDTt27EBhYSF++OEHjBo1CqtXr/Z2ta4qJ0+exFNPPYXs7Gz4+fl5uzpXtYEDB5q/b9u2Lbp164akpCR899138Pf392LNGo49KjVERkZCLpfbzIA+d+4cYmNjvVSrxqWqnWprw9jYWOTl5Vk8X1lZiYsXL1qUsXeMmudwVOZq+lk98cQT+Pnnn7Fq1SokJCSY98fGxqKiogIFBQUW5a3buaFtqNFo4O/vf038TahUKqSmpqJTp06YPHky2rVrh+nTp7ONXWjr1q3Iy8tDx44doVAooFAosHr1anzwwQdQKBSIiYlhW7tJaGgoWrZsicOHDzfa32kGlRpUKhU6deqE33//3bzPaDTi999/R1ZWlhdr1ng0a9YMsbGxFm2o1WqxceNGcxtmZWWhoKAAW7duNZdZuXIljEYjunXrZi6zZs0a6PV6c5ns7GykpaUhLCzMXKbmearKXA0/KyEEnnjiCSxcuBArV65Es2bNLJ7v1KkTlEqlxfs/cOAATpw4YdHOu3fvtgiF2dnZ0Gg0yMjIMJeprQ2vxb8Jo9EInU7HNnahPn36YPfu3dixY4f50blzZ4wcOdL8PdvaPYqLi3HkyBE0adKk8f5O13v67VVu/vz5Qq1Wizlz5oi//vpLPPLIIyI0NNRiBvS1rqioSGzfvl1s375dABDTpk0T27dvF8ePHxdCmC5PDg0NFYsXLxa7du0SQ4cOtXt5cocOHcTGjRvFunXrRIsWLSwuTy4oKBAxMTHivvvuE3v27BHz588XAQEBNpcnKxQK8Z///Efs27dPvPbaa1fN5cn//Oc/RUhIiMjJybG4zLC0tNRc5rHHHhNNmzYVK1euFFu2bBFZWVkiKyvL/HzVZYb9+/cXO3bsEMuXLxdRUVF2LzN87rnnxL59+8THH39s9zLDq/Vv4l//+pdYvXq1yM3NFbt27RL/+te/hCRJ4rfffhNCsI3dqeZVP0KwrV3lmWeeETk5OSI3N1f88ccfom/fviIyMlLk5eUJIRpnOzOo2PHhhx+Kpk2bCpVKJbp27So2bNjg7Sr5lFWrVgkANo9Ro0YJIUyXKL/yyisiJiZGqNVq0adPH3HgwAGLY1y4cEGMGDFCBAUFCY1GI8aMGSOKioosyuzcuVPccMMNQq1Wi/j4eDFlyhSbunz33XeiZcuWQqVSidatW4tffvnFbe/bk+y1LwAxe/Zsc5mysjLx+OOPi7CwMBEQECBuu+02cebMGYvjHDt2TAwcOFD4+/uLyMhI8cwzzwi9Xm9RZtWqVaJ9+/ZCpVKJ5s2bW5yjytX6N/HAAw+IpKQkoVKpRFRUlOjTp485pAjBNnYn66DCtnaN4cOHiyZNmgiVSiXi4+PF8OHDxeHDh83PN8Z2loQQov79MERERETuxzkqRERE5LMYVIiIiMhnMagQERGRz2JQISIiIp/FoEJEREQ+i0GFiIiIfBaDChEREfksBhUiIiLyWQwqRORRo0ePxrBhw7xdDSJqJBhUiIiIyGcxqBCRW/zwww/IzMyEv78/IiIi0LdvXzz33HP46quvsHjxYkiSBEmSkJOTAwA4efIk7rrrLoSGhiI8PBxDhw7FsWPHzMer6omZOHEioqKioNFo8Nhjj6GioqLWc5aUlHj4nRORKym8XQEiuvqcOXMGI0aMwNSpU3HbbbehqKgIa9euxf33348TJ05Aq9Vi9uzZAIDw8HDo9XoMGDAAWVlZWLt2LRQKBd544w3cfPPN2LVrF1QqFQDg999/h5+fH3JycnDs2DGMGTMGERERePPNNx2ek7czI2rcGFSIyOXOnDmDyspK3H777UhKSgIAZGZmAgD8/f2h0+kQGxtrLv/f//4XRqMRs2bNgiRJAIDZs2cjNDQUOTk56N+/PwBApVLhyy+/REBAAFq3bo3XX38dzz33HCZNmlTrOYmo8eLQDxG5XLt27dCnTx9kZmbizjvvxOeff45Lly45LL9z504cPnwYwcHBCAoKQlBQEMLDw1FeXo4jR45YHDcgIMC8nZWVheLiYpw8ebLe5ySixoFBhYhcTi6XIzs7G8uWLUNGRgY+/PBDpKWlITc312754uJidOrUCTt27LB4HDx4EPfcc49bzklEjQODChG5hSRJ6N69OyZOnIjt27dDpVJh4cKFUKlUMBgMFmU7duyIQ4cOITo6GqmpqRaPkJAQc7mdO3eirKzMvL1hwwYEBQUhMTGx1nMSUePFoEJELrdx40a89dZb2LJlC06cOIEFCxbg/PnzaNWqFZKTk7Fr1y4cOHAA+fn50Ov1GDlyJCIjIzF06FCsXbsWubm5yMnJwZNPPom///7bfNyKigo8+OCD+Ouvv7B06VK89tpreOKJJyCTyWo9JxE1XpxMS0Qup9FosGbNGrz//vvQarVISkrCu+++i4EDB6Jz587IyclB586dUVxcjFWrVqF3795Ys2YNXnjhBdx+++0oKipCfHw8+vTpA41GYz5unz590KJFC/Ts2RM6nQ4jRozAhAkT6jwnETVekuC1e0TUCIwePRoFBQVYtGiRt6tCRB7EoR8iIiLyWQwqRERE5LM49ENEREQ+iz0qRERE5LMYVIiIiMhnMagQERGRz2JQISIiIp/FoEJEREQ+i0GFiIiIfBaDChEREfksBhUiIiLyWQwqRERE5LP+H+tD8l6LrH6DAAAAAElFTkSuQmCC\n" + }, + "metadata": {} + } + ], + "source": [ + "plt.title(f\"Convergence of adamw (train loss)\")\n", + "plt.plot(all_train_losses, label=\"train\", lw=3)\n", + "plt.plot(\n", + " jnp.arange(0, len(all_eval_losses) * N_FREQ_EVAL, N_FREQ_EVAL),\n", + " all_eval_losses,\n", + " label=\"test\",\n", + " lw=3,\n", + ")\n", + "plt.xlabel(\"steps\")\n", + "plt.ylabel(\"loss\")\n", + "plt.grid()\n", + "plt.legend(frameon=False)\n", + "plt.show()" + ] + }, + { + "cell_type": "markdown", + "metadata": { + "id": "E6-aaLDL7RbI" + }, + "source": [ + "# Text generation\n", + "\n", + "Finally, after training, we use the generate function to let the NanoGPT model demonstrate its ability to create text that resembles Shakespeare, albeit in a miniature form." + ] + }, + { + "cell_type": "code", + "execution_count": null, + "metadata": { + "colab": { + "base_uri": "https://localhost:8080/" + }, + "id": "6AejKtZnFmhK", + "outputId": "78e7b13f-28ff-4786-e206-bcd104e3c507" + }, + "outputs": [ + { + "output_type": "stream", + "name": "stdout", + "text": [ + "GREMIO:\n", + "So called, officer:\n", + "Peace, for me would be a mighty heart, and I'll do\n", + "any good.\n", + "\n", + "MERCUTIO:\n", + "Ay, look'd dead made.\n", + "\n", + "VALERIA:\n", + "Rather is in the present book of love\n", + "Than piece thy thoughts to shame, but yet book known\n", + "Where my Edward, whom I so could do it to him;\n", + "And see how the government of your blood,\n", + "Congeal'd your kingdom then we resign your highness;\n", + "Unless you scarcely know how I came, his wife;\n", + "And I may contain a little of England's heart\n", + "After a king of disdain! supect, tribunes,\n", + "The sun under thine is too large:\n", + "The senators do add to their trenches stretch'd,\n", + "With such as a toy, or be it known, I would\n", + "See my shame home, sweet friends, we are too dear:\n", + "Tell me consul, what's become the soldiers?\n", + "\n", + "SOMERSET:\n", + "So fled; alas! much is possess'd by this means?\n", + "Away with the nightingale. He hast he wounded his sceptre\n", + "And hide his heir: go hence, to have his hearts\n", + "To close our bloods, for the climate to be\n", + "Ere one would be so often been seen, a greater stranger\n", + "Can have you \n" + ] + } + ], + "source": [ + "# Let's now generate some text\n", + "key, subkey = jax.random.split(key)\n", + "text = model.generate(key, var_params, 1000)[:, 0, 0].tolist()\n", + "print(decode(text))" + ] + } + ], + "metadata": { + "accelerator": "GPU", + "colab": { + "gpuType": "A100", + "machine_shape": "hm", + "provenance": [] + }, + "kernelspec": { + "display_name": "Python 3", + "name": "python3" + }, + "language_info": { + "name": "python" + } + }, + "nbformat": 4, + "nbformat_minor": 0 +}