{ "cells": [ { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "# default_exp learner" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#export\n", "from fastai2.data.all import *\n", "from fastai2.optimizer import *\n", "from fastai2.callback.core import *" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#hide\n", "from nbdev.showdoc import *" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#export\n", "_all_ = ['CancelFitException', 'CancelEpochException', 'CancelTrainException', 'CancelValidException', 'CancelBatchException']" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#export\n", "_loop = ['Start Fit', 'before_fit', 'Start Epoch Loop', 'before_epoch', 'Start Train', 'before_train',\n", " 'Start Batch Loop', 'before_batch', 'after_pred', 'after_loss', 'before_backward', 'after_backward',\n", " 'after_step', 'after_cancel_batch', 'after_batch','End Batch Loop','End Train',\n", " 'after_cancel_train', 'after_train', 'Start Valid', 'before_validate','Start Batch Loop',\n", " '**CBs same as train batch**', 'End Batch Loop', 'End Valid', 'after_cancel_validate',\n", " 'after_validate', 'End Epoch Loop', 'after_cancel_epoch', 'after_epoch', 'End Fit',\n", " 'after_cancel_fit', 'after_fit']" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "# Learner\n", "\n", "> Basic class for handling the training loop" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "You probably want to jump directly to the definition of `Learner`." ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "## Utils function" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#hide\n", "#For tests\n", "from torch.utils.data import TensorDataset\n", "\n", "def synth_dbunch(a=2, b=3, bs=16, n_train=10, n_valid=2, cuda=False):\n", " \"A simple dataset where `x` is random and `y = a*x + b` plus some noise.\"\n", " def get_data(n):\n", " x = torch.randn(int(bs*n))\n", " return TensorDataset(x, a*x + b + 0.1*torch.randn(int(bs*n)))\n", " train_ds = get_data(n_train)\n", " valid_ds = get_data(n_valid)\n", " device = default_device() if cuda else None\n", " train_dl = TfmdDL(train_ds, bs=bs, shuffle=True, num_workers=0)\n", " valid_dl = TfmdDL(valid_ds, bs=bs, num_workers=0)\n", " return DataLoaders(train_dl, valid_dl, device=device)\n", "\n", "class RegModel(Module):\n", " \"A r\"\n", " def __init__(self): self.a,self.b = nn.Parameter(torch.randn(1)),nn.Parameter(torch.randn(1))\n", " def forward(self, x): return x*self.a + self.b" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "# export\n", "defaults.lr = 1e-3" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "# export\n", "def replacing_yield(o, attr, val):\n", " \"Context manager to temporarily replace an attribute\"\n", " old = getattr(o,attr)\n", " try: yield setattr(o,attr,val)\n", " finally: setattr(o,attr,old)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "class _A:\n", " def __init__(self, a): self.a = a\n", " @contextmanager\n", " def a_changed(self, v): return replacing_yield(self, 'a', v)\n", "\n", "a = _A(42)\n", "with a.a_changed(32):\n", " test_eq(a.a, 32)\n", "test_eq(a.a, 42)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#export\n", "def mk_metric(m):\n", " \"Convert `m` to an `AvgMetric`, unless it's already a `Metric`\"\n", " return m if isinstance(m, Metric) else AvgMetric(m)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "See the class `Metric` below for more information." ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#export\n", "def save_model(file, model, opt, with_opt=True, pickle_protocol=2):\n", " \"Save `model` to `file` along with `opt` (if available, and if `with_opt`)\"\n", " if rank_distrib(): return # don't save if child proc\n", " if opt is None: with_opt=False\n", " state = get_model(model).state_dict()\n", " if with_opt: state = {'model': state, 'opt':opt.state_dict()}\n", " torch.save(state, file, pickle_protocol=pickle_protocol)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "`file` can be a `Path` object, a string or an opened file object. `pickle_protocol` is passed along to `torch.save`" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "# export\n", "def load_model(file, model, opt, with_opt=None, device=None, strict=True):\n", " \"Load `model` from `file` along with `opt` (if available, and if `with_opt`)\"\n", " distrib_barrier()\n", " if isinstance(device, int): device = torch.device('cuda', device)\n", " elif device is None: device = 'cpu'\n", " state = torch.load(file, map_location=device)\n", " hasopt = set(state)=={'model', 'opt'}\n", " model_state = state['model'] if hasopt else state\n", " get_model(model).load_state_dict(model_state, strict=strict)\n", " if hasopt and ifnone(with_opt,True):\n", " try: opt.load_state_dict(state['opt'])\n", " except:\n", " if with_opt: warn(\"Could not load the optimizer state.\")\n", " elif with_opt: warn(\"Saved filed doesn't contain an optimizer state.\")" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "`file` can be a `Path` object, a string or an opened file object. If a `device` is passed, the model is loaded on it, otherwise it's loaded on the CPU. \n", "\n", "If `strict` is `True`, the file must exactly contain weights for every parameter key in `model`, if `strict` is `False`, only the keys that are in the saved model are loaded in `model`." ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "# export\n", "def _try_concat(o):\n", " try: return torch.cat(o)\n", " except: return sum([L(o_[i,:] for i in range_of(o_)) for o_ in o], L())" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#export\n", "_before_epoch = [event.before_fit, event.before_epoch]\n", "_after_epoch = [event.after_epoch, event.after_fit]" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#export\n", "class _ConstantFunc():\n", " \"Returns a function that returns `o`\"\n", " def __init__(self, o): self.o = o\n", " def __call__(self, *args, **kwargs): return self.o" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "## Learner -" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "# export\n", "@log_args(but='dls,model,opt_func,cbs')\n", "class Learner():\n", " def __init__(self, dls, model, loss_func=None, opt_func=Adam, lr=defaults.lr, splitter=trainable_params, cbs=None,\n", " metrics=None, path=None, model_dir='models', wd=None, wd_bn_bias=False, train_bn=True,\n", " moms=(0.95,0.85,0.95)):\n", " store_attr(self, \"dls,model,opt_func,lr,splitter,model_dir,wd,wd_bn_bias,train_bn,metrics,moms\")\n", " self.training,self.create_mbar,self.logger,self.opt,self.cbs = False,True,print,None,L()\n", " if loss_func is None:\n", " loss_func = getattr(dls.train_ds, 'loss_func', None)\n", " assert loss_func is not None, \"Could not infer loss function from the data, please pass a loss function.\"\n", " self.loss_func = loss_func\n", " self.path = Path(path) if path is not None else getattr(dls, 'path', Path('.'))\n", " self.add_cbs([(cb() if isinstance(cb, type) else cb) for cb in L(defaults.callbacks)+L(cbs)])\n", " self.epoch,self.n_epoch,self.loss = 0,1,tensor(0.)\n", "\n", " @property\n", " def metrics(self): return self._metrics\n", " @metrics.setter\n", " def metrics(self,v): self._metrics = L(v).map(mk_metric)\n", "\n", " def _grab_cbs(self, cb_cls): return L(cb for cb in self.cbs if isinstance(cb, cb_cls))\n", " def add_cbs(self, cbs): L(cbs).map(self.add_cb)\n", " def remove_cbs(self, cbs): L(cbs).map(self.remove_cb)\n", " def add_cb(self, cb):\n", " old = getattr(self, cb.name, None)\n", " assert not old or isinstance(old, type(cb)), f\"self.{cb.name} already registered\"\n", " cb.learn = self\n", " setattr(self, cb.name, cb)\n", " self.cbs.append(cb)\n", " return self\n", "\n", " def remove_cb(self, cb):\n", " if isinstance(cb, type): self.remove_cbs(self._grab_cbs(cb))\n", " else:\n", " cb.learn = None\n", " if hasattr(self, cb.name): delattr(self, cb.name)\n", " if cb in self.cbs: self.cbs.remove(cb)\n", "\n", " @contextmanager\n", " def added_cbs(self, cbs):\n", " self.add_cbs(cbs)\n", " try: yield\n", " finally: self.remove_cbs(cbs)\n", "\n", " @contextmanager\n", " def removed_cbs(self, cbs):\n", " self.remove_cbs(cbs)\n", " try: yield self\n", " finally: self.add_cbs(cbs)\n", "\n", " def ordered_cbs(self, event): return [cb for cb in sort_by_run(self.cbs) if hasattr(cb, event)]\n", "\n", " def __call__(self, event_name): L(event_name).map(self._call_one)\n", "\n", " def _call_one(self, event_name):\n", " assert hasattr(event, event_name), event_name\n", " [cb(event_name) for cb in sort_by_run(self.cbs)]\n", "\n", " def _bn_bias_state(self, with_bias): return norm_bias_params(self.model, with_bias).map(self.opt.state)\n", " def create_opt(self):\n", " self.opt = self.opt_func(self.splitter(self.model), lr=self.lr)\n", " if not self.wd_bn_bias:\n", " for p in self._bn_bias_state(True ): p['do_wd'] = False\n", " if self.train_bn:\n", " for p in self._bn_bias_state(False): p['force_train'] = True\n", "\n", " def _split(self, b):\n", " i = getattr(self.dls, 'n_inp', 1 if len(b)==1 else len(b)-1)\n", " self.xb,self.yb = b[:i],b[i:]\n", "\n", " def _step(self): self.opt.step()\n", " def _backward(self): self.loss.backward()\n", "\n", " def _with_events(self, f, event_type, ex, final=noop):\n", " try: self(f'before_{event_type}') ;f()\n", " except ex: self(f'after_cancel_{event_type}')\n", " finally: self(f'after_{event_type}') ;final()\n", "\n", " def all_batches(self):\n", " self.n_iter = len(self.dl)\n", " for o in enumerate(self.dl): self.one_batch(*o)\n", "\n", " def _do_one_batch(self):\n", " self.pred = self.model(*self.xb); self('after_pred')\n", " if len(self.yb) == 0: return\n", " self.loss = self.loss_func(self.pred, *self.yb); self('after_loss')\n", " if not self.training: return\n", " self('before_backward')\n", " self._backward(); self('after_backward')\n", " self._step(); self('after_step')\n", " self.opt.zero_grad()\n", "\n", " def one_batch(self, i, b):\n", " self.iter = i\n", " self._split(b)\n", " self._with_events(self._do_one_batch, 'batch', CancelBatchException)\n", "\n", " def _do_epoch_train(self):\n", " self.dl = self.dls.train\n", " self._with_events(self.all_batches, 'train', CancelTrainException)\n", "\n", " def _do_epoch_validate(self, ds_idx=1, dl=None):\n", " if dl is None: dl = self.dls[ds_idx]\n", " self.dl = dl;\n", " with torch.no_grad(): self._with_events(self.all_batches, 'validate', CancelValidException)\n", "\n", " def _do_epoch(self):\n", " self._do_epoch_train()\n", " self._do_epoch_validate()\n", "\n", " def _do_fit(self):\n", " for epoch in range(self.n_epoch):\n", " self.epoch=epoch\n", " self._with_events(self._do_epoch, 'epoch', CancelEpochException)\n", "\n", " @log_args(but='cbs')\n", " def fit(self, n_epoch, lr=None, wd=None, cbs=None, reset_opt=False):\n", " with self.added_cbs(cbs):\n", " if reset_opt or not self.opt: self.create_opt()\n", " if wd is None: wd = self.wd\n", " if wd is not None: self.opt.set_hypers(wd=wd)\n", " self.opt.set_hypers(lr=self.lr if lr is None else lr)\n", " self.n_epoch,self.loss = n_epoch,tensor(0.)\n", " self._with_events(self._do_fit, 'fit', CancelFitException, self._end_cleanup)\n", "\n", " def _end_cleanup(self): self.dl,self.xb,self.yb,self.pred,self.loss = None,(None,),(None,),None,None\n", " def __enter__(self): self(_before_epoch); return self\n", " def __exit__(self, exc_type, exc_value, tb): self(_after_epoch)\n", "\n", " def validation_context(self, cbs=None, inner=False):\n", " cms = [self.no_logging(),self.no_mbar()]\n", " if cbs: cms.append(self.added_cbs(cbs))\n", " if not inner: cms.append(self)\n", " return ContextManagers(cms)\n", "\n", " def validate(self, ds_idx=1, dl=None, cbs=None):\n", " if dl is None: dl = self.dls[ds_idx]\n", " with self.validation_context(cbs=cbs): self._do_epoch_validate(ds_idx, dl)\n", " return getattr(self, 'final_record', None)\n", "\n", " @delegates(GatherPredsCallback.__init__)\n", " def get_preds(self, ds_idx=1, dl=None, with_input=False, with_decoded=False, with_loss=False, act=None,\n", " inner=False, reorder=True, cbs=None, **kwargs):\n", " if dl is None: dl = self.dls[ds_idx].new(shuffled=False, drop_last=False)\n", " if reorder and hasattr(dl, 'get_idxs'):\n", " idxs = dl.get_idxs()\n", " dl = dl.new(get_idxs = _ConstantFunc(idxs))\n", " cb = GatherPredsCallback(with_input=with_input, with_loss=with_loss, **kwargs)\n", " ctx_mgrs = self.validation_context(cbs=L(cbs)+[cb], inner=inner)\n", " if with_loss: ctx_mgrs.append(self.loss_not_reduced())\n", " with ContextManagers(ctx_mgrs):\n", " self._do_epoch_validate(dl=dl)\n", " if act is None: act = getattr(self.loss_func, 'activation', noop)\n", " res = cb.all_tensors()\n", " pred_i = 1 if with_input else 0\n", " if res[pred_i] is not None:\n", " res[pred_i] = act(res[pred_i])\n", " if with_decoded: res.insert(pred_i+2, getattr(self.loss_func, 'decodes', noop)(res[pred_i]))\n", " if reorder and hasattr(dl, 'get_idxs'): res = nested_reorder(res, tensor(idxs).argsort())\n", " return tuple(res)\n", " self._end_cleanup()\n", "\n", " def predict(self, item, rm_type_tfms=None, with_input=False):\n", " dl = self.dls.test_dl([item], rm_type_tfms=rm_type_tfms, num_workers=0)\n", " inp,preds,_,dec_preds = self.get_preds(dl=dl, with_input=True, with_decoded=True)\n", " i = getattr(self.dls, 'n_inp', -1)\n", " inp = (inp,) if i==1 else tuplify(inp)\n", " dec = self.dls.decode_batch(inp + tuplify(dec_preds))[0]\n", " dec_inp,dec_targ = map(detuplify, [dec[:i],dec[i:]])\n", " res = dec_targ,dec_preds[0],preds[0]\n", " if with_input: res = (dec_inp,) + res\n", " return res\n", "\n", " def show_results(self, ds_idx=1, dl=None, max_n=9, shuffle=True, **kwargs):\n", " if dl is None: dl = self.dls[ds_idx].new(shuffle=shuffle)\n", " b = dl.one_batch()\n", " _,_,preds = self.get_preds(dl=[b], with_decoded=True)\n", " self.dls.show_results(b, preds, max_n=max_n, **kwargs)\n", "\n", " def show_training_loop(self):\n", " indent = 0\n", " for s in _loop:\n", " if s.startswith('Start'): print(f'{\" \"*indent}{s}'); indent += 2\n", " elif s.startswith('End'): indent -= 2; print(f'{\" \"*indent}{s}')\n", " else: print(f'{\" \"*indent} - {s:15}:', self.ordered_cbs(s))\n", "\n", " @contextmanager\n", " def no_logging(self): return replacing_yield(self, 'logger', noop)\n", " @contextmanager\n", " def no_mbar(self): return replacing_yield(self, 'create_mbar', False)\n", "\n", " @contextmanager\n", " def loss_not_reduced(self):\n", " if hasattr(self.loss_func, 'reduction'): return replacing_yield(self.loss_func, 'reduction', 'none')\n", " else: return replacing_yield(self, 'loss_func', partial(self.loss_func, reduction='none'))\n", "\n", " @delegates(save_model)\n", " def save(self, file, **kwargs):\n", " file = join_path_file(file, self.path/self.model_dir, ext='.pth')\n", " save_model(file, self.model, getattr(self,'opt',None), **kwargs)\n", "\n", " @delegates(load_model)\n", " def load(self, file, with_opt=None, device=None, **kwargs):\n", " if device is None and hasattr(self.dls, 'device'): device = self.dls.device\n", " if self.opt is None: self.create_opt()\n", " file = join_path_file(file, self.path/self.model_dir, ext='.pth')\n", " load_model(file, self.model, self.opt, device=device, **kwargs)\n", " return self\n", "\n", "Learner.x,Learner.y = add_props(lambda i,x: detuplify((x.xb,x.yb)[i]))" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#export\n", "add_docs(Learner, \"Group together a `model`, some `dls` and a `loss_func` to handle training\",\n", " add_cbs=\"Add `cbs` to the list of `Callback` and register `self` as their learner\",\n", " add_cb=\"Add `cb` to the list of `Callback` and register `self` as their learner\",\n", " remove_cbs=\"Remove `cbs` from the list of `Callback` and deregister `self` as their learner\",\n", " remove_cb=\"Add `cb` from the list of `Callback` and deregister `self` as their learner\",\n", " added_cbs=\"Context manage that temporarily adds `cbs`\",\n", " removed_cbs=\"Context manage that temporarily removes `cbs`\",\n", " ordered_cbs=\"List of `Callback`s, in order, for an `event` in the training loop\",\n", " create_opt=\"Create an optimizer with default hyper-parameters\",\n", " one_batch=\"Train or evaluate `self.model` on batch `(xb,yb)`\",\n", " all_batches=\"Train or evaluate `self.model` on all the batches of `self.dl`\",\n", " fit=\"Fit `self.model` for `n_epoch` using `cbs`. Optionally `reset_opt`.\",\n", " validate=\"Validate on `dl` with potential new `cbs`.\",\n", " get_preds=\"Get the predictions and targets on the `ds_idx`-th dbunchset or `dl`, optionally `with_input` and `with_loss`\",\n", " predict=\"Prediction on `item`, fully decoded, loss function decoded and probabilities\",\n", " validation_context=\"A `ContextManagers` suitable for validation, with optional `cbs`\",\n", " show_results=\"Show some predictions on `ds_idx`-th dataset or `dl`\",\n", " show_training_loop=\"Show each step in the training loop\",\n", " no_logging=\"Context manager to temporarily remove `logger`\",\n", " no_mbar=\"Context manager to temporarily prevent the master progress bar from being created\",\n", " loss_not_reduced=\"A context manager to evaluate `loss_func` with reduction set to none.\",\n", " save=\"Save model and optimizer state (if `with_opt`) to `self.path/self.model_dir/file`\",\n", " load=\"Load model and optimizer state (if `with_opt`) from `self.path/self.model_dir/file` using `device`\",\n", " __call__=\"Call `event_name` for all `Callback`s in `self.cbs`\"\n", ")" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

class Learner[source]

\n", "\n", "> Learner(**`dls`**, **`model`**, **`loss_func`**=*`None`*, **`opt_func`**=*`'Adam'`*, **`lr`**=*`0.001`*, **`splitter`**=*`'trainable_params'`*, **`cbs`**=*`None`*, **`metrics`**=*`None`*, **`path`**=*`None`*, **`model_dir`**=*`'models'`*, **`wd`**=*`None`*, **`wd_bn_bias`**=*`False`*, **`train_bn`**=*`True`*, **`moms`**=*`(0.95, 0.85, 0.95)`*)\n", "\n", "Group together a `model`, some `dls` and a `loss_func` to handle training" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(Learner)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "`opt_func` will be used to create an optimizer when `Learner.fit` is called, with `lr` as a default learning rate. `splitter` is a function that takes `self.model` and returns a list of parameter groups (or just one parameter group if there are no different parameter groups). The default is `trainable_params`, which returns all trainable parameters of the model.\n", "\n", "`cbs` is one or a list of `Callback`s to pass to the `Learner`. `Callback`s are used for every tweak of the training loop. Each `Callback` is registered as an attribute of `Learner` (with camel case). At creation, all the callbacks in `defaults.callbacks` (`TrainEvalCallback`, `Recorder` and `ProgressCallback`) are associated to the `Learner`.\n", "\n", "`metrics` is an optional list of metrics, that can be either functions or `Metric`s (see below). \n", "\n", "`path` and `model_dir` are used to save and/or load models. Often `path` will be inferred from `dls`, but you can override it or pass a `Path` object to `model_dir`. Make sure you can write in `path/model_dir`!\n", "\n", "`wd` is the default weight decay used when training the model; `moms`, the default momentums used in `Learner.fit_one_cycle`. `wd_bn_bias` controls if weight decay is applied to `BatchNorm` layers and bias. \n", "\n", "Lastly, `train_bn` controls if `BatchNorm` layers are trained even when they are supposed to be frozen according to the `splitter`. Our empirical experiments have shown that it's the best behavior for those layers in transfer learning." ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "### PyTorch interrop" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "You can use regular PyTorch functionality for most of the arguments of the `Learner`, although the experience will be smoother with pure fastai objects and you will be able to use the full functionality of the library. The expectation is that the training loop will work smoothly even if you did not use fastai end to end. What you might lose are interpretation objects or showing functionality. The list below explains how to use plain PyTorch objects for all the arguments and what you might lose.\n", "\n", "The most important is `opt_func`. If you are not using a fastai optimizer, you will need to write a function that wraps your PyTorch optimizer in an `OptimWrapper`. See the [optimizer module](http://dev.fast.ai/optimizer) for more details. This is to ensure the library's schedulers/freeze API work with your code.\n", "\n", "- `dls` is a `DataLoaders` object, that you can create from standard PyTorch dataloaders. By doing so, you will lose all showing functionality like `show_batch`/`show_results`. You can check the [data block API](http://dev.fast.ai/tutorial.datablock) or the [mid-level data API tutorial](http://dev.fast.ai/tutorial.pets) to learn how to use fastai to gather your data!\n", "- `model` is a standard PyTorch model. You can use anyone you like, just make sure it accepts the number of inputs you have in your `DataLoaders` and returns as many outputs as you have targets.\n", "- `loss_func` can be any loss function you like. It needs to be one of fastai's if you want to use `Learn.predict` or `Learn.get_preds`, or you will have to implement special methods (see more details after the `BaseLoss` documentation)." ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "Now let's look at the main thing the `Learner` class implements: the training loop." ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "### Training loop" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

Learner.fit[source]

\n", "\n", "> Learner.fit(**`n_epoch`**, **`lr`**=*`None`*, **`wd`**=*`None`*, **`cbs`**=*`None`*, **`reset_opt`**=*`False`*)\n", "\n", "Fit `self.model` for `n_epoch` using `cbs`. Optionally `reset_opt`." ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(Learner.fit)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "Uses `lr` and `wd` if they are provided, otherwise use the defaults values given by the `lr` and `wd` attributes of `Learner`." ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "All the examples use `synth_learner` which is a simple `Learner` training a linear regression model." ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#hide\n", "def synth_learner(n_train=10, n_valid=2, cuda=False, lr=defaults.lr, **kwargs):\n", " data = synth_dbunch(n_train=n_train,n_valid=n_valid, cuda=cuda)\n", " return Learner(data, RegModel(), loss_func=MSELossFlat(), lr=lr, **kwargs)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#Training a few epochs should make the model better\n", "learn = synth_learner(lr=5e-2)\n", "learn.model = learn.model.cpu()\n", "xb,yb = learn.dls.one_batch()\n", "init_loss = learn.loss_func(learn.model(xb), yb)\n", "learn.fit(6)\n", "xb,yb = learn.dls.one_batch()\n", "final_loss = learn.loss_func(learn.model(xb), yb)\n", "assert final_loss < init_loss" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#hide\n", "#Test of TrainEvalCallback\n", "class TestTrainEvalCallback(Callback):\n", " run_after,run_valid = TrainEvalCallback,False\n", " def before_fit(self): \n", " test_eq([self.pct_train,self.train_iter], [0., 0])\n", " self.old_pct_train,self.old_train_iter = self.pct_train,self.train_iter\n", " \n", " def before_batch(self): test_eq(next(self.model.parameters()).device, find_device(self.xb))\n", " \n", " def after_batch(self):\n", " assert self.training\n", " test_eq(self.pct_train , self.old_pct_train+1/(self.n_iter*self.n_epoch))\n", " test_eq(self.train_iter, self.old_train_iter+1)\n", " self.old_pct_train,self.old_train_iter = self.pct_train,self.train_iter\n", " \n", " def before_train(self):\n", " assert self.training and self.model.training\n", " test_eq(self.pct_train, self.epoch/self.n_epoch)\n", " self.old_pct_train = self.pct_train\n", " \n", " def before_validate(self):\n", " assert not self.training and not self.model.training\n", " \n", "learn = synth_learner(cbs=TestTrainEvalCallback)\n", "learn.fit(1)\n", "#Check order is properly taken into account\n", "learn.cbs = L(reversed(learn.cbs))" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#hide\n", "#cuda\n", "#Check model is put on the GPU if needed\n", "learn = synth_learner(cbs=TestTrainEvalCallback, cuda=True)\n", "learn.fit(1)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#hide\n", "#Check wd is not applied on bn/bias when option wd_bn_bias=False\n", "class _TstModel(nn.Module):\n", " def __init__(self):\n", " super().__init__()\n", " self.a,self.b = nn.Parameter(torch.randn(1)),nn.Parameter(torch.randn(1))\n", " self.tst = nn.Sequential(nn.Linear(4,5), nn.BatchNorm1d(3))\n", " self.tst[0].bias.data,self.tst[1].bias.data = torch.randn(5),torch.randn(3) \n", " def forward(self, x): return x * self.a + self.b\n", " \n", "class _PutGrad(Callback):\n", " def after_backward(self):\n", " for p in self.learn.model.tst.parameters():\n", " p.grad = torch.ones_like(p.data)\n", " \n", "learn = synth_learner(n_train=5, opt_func = partial(SGD, wd=1, decouple_wd=True), cbs=_PutGrad)\n", "learn.model = _TstModel()\n", "init = [p.clone() for p in learn.model.tst.parameters()]\n", "learn.fit(1, lr=1e-2)\n", "end = list(learn.model.tst.parameters())\n", "for i in [0]: assert not torch.allclose(end[i]-init[i], -0.05 * torch.ones_like(end[i]))\n", "for i in [1,2,3]: test_close(end[i]-init[i], -0.05 * torch.ones_like(end[i]))" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

Learner.one_batch[source]

\n", "\n", "> Learner.one_batch(**`i`**, **`b`**)\n", "\n", "Train or evaluate `self.model` on batch `(xb,yb)`" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(Learner.one_batch)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "This is an internal method called by `Learner.fit`. If passed, `i` is the index of this iteration in the epoch. In training mode, this does a full training step on the batch (compute predictions, loss, gradients, update the model parameters and zero the gradients). In validation mode, it stops at the loss computation. Training or validation is controlled internally by the `TrainEvalCallback` through the `training` attribute.\n", "\n", "Nothing is returned, but the attributes `x`, `y`, `pred`, `loss` of the `Learner` are set with the propet values:" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "b = learn.dls.one_batch()\n", "learn.one_batch(0, b)\n", "test_eq(learn.x, b[0])\n", "test_eq(learn.y, b[1])\n", "out = learn.model(learn.x)\n", "test_eq(learn.pred, out)\n", "test_eq(learn.loss, learn.loss_func(out, b[1]))" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "More generally, the following attributes of `Learner` are available and updated during the training loop:\n", "- `model`: the model used for training/validation\n", "- `data`: the underlying `DataLoaders`\n", "- `loss_func`: the loss function used\n", "- `opt`: the optimizer used to udpate the model parameters\n", "- `opt_func`: the function used to create the optimizer\n", "- `cbs`: the list containing all `Callback`s\n", "- `dl`: current `DataLoader` used for iteration\n", "- `x`/`xb`: last input drawn from `self.dl` (potentially modified by callbacks). `xb` is always a tuple (potentially with one element) and `x` is detuplified. You can only assign to `xb`.\n", "- `y`/`yb`: last target drawn from `self.dl` (potentially modified by callbacks). `yb` is always a tuple (potentially with one element) and `y` is detuplified. You can only assign to `yb`.\n", "- `pred`: last predictions from `self.model` (potentially modified by callbacks)\n", "- `loss`: last computed loss (potentially modified by callbacks)\n", "- `n_epoch`: the number of epochs in this training\n", "- `n_iter`: the number of iterations in the current `self.dl`\n", "- `epoch`: the current epoch index (from 0 to `n_epoch-1`)\n", "- `iter`: the current iteration index in `self.dl` (from 0 to `n_iter-1`)\n", "\n", "The following attributes are added by `TrainEvalCallback` and should be available unless you went out of your way to remove that callback:\n", "\n", "- `train_iter`: the number of training iterations done since the beginning of this training\n", "- `pct_train`: from 0. to 1., the percentage of training iterations completed\n", "- `training`: flag to indicate if we're in training mode or not\n", "\n", "The following attribute is added by `Recorder` and should be available unless you went out of your way to remove that callback:\n", "\n", "- `smooth_loss`: an exponentially-averaged version of the training loss" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#hide\n", "class VerboseCallback(Callback):\n", " \"Callback that prints the name of each event called\"\n", " def __call__(self, event_name):\n", " print(event_name)\n", " super().__call__(event_name)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#hide\n", "class TestOneBatch(VerboseCallback):\n", " def __init__(self, xb, yb, i):\n", " self.save_xb,self.save_yb,self.i = xb,yb,i\n", " self.old_pred,self.old_loss = None,tensor(0.)\n", " \n", " def before_batch(self):\n", " self.old_a,self.old_b = self.model.a.data.clone(),self.model.b.data.clone()\n", " test_eq(self.iter, self.i)\n", " test_eq(self.save_xb, *self.xb)\n", " test_eq(self.save_yb, *self.yb)\n", " if hasattr(self.learn, 'pred'): test_eq(self.pred, self.old_pred)\n", " \n", " def after_pred(self):\n", " self.old_pred = self.pred\n", " test_eq(self.pred, self.model.a.data * self.x + self.model.b.data)\n", " test_eq(self.loss, self.old_loss)\n", " \n", " def after_loss(self):\n", " self.old_loss = self.loss\n", " test_eq(self.loss, self.loss_func(self.old_pred, self.save_yb))\n", " for p in self.model.parameters(): \n", " if not hasattr(p, 'grad') or p.grad is not None: test_eq(p.grad, tensor([0.]))\n", " \n", " def after_backward(self):\n", " self.grad_a = (2 * self.x * (self.pred.data - self.y)).mean()\n", " self.grad_b = 2 * (self.pred.data - self.y).mean()\n", " test_close(self.model.a.grad.data, self.grad_a)\n", " test_close(self.model.b.grad.data, self.grad_b)\n", " test_eq(self.model.a.data, self.old_a)\n", " test_eq(self.model.b.data, self.old_b)\n", " \n", " def after_step(self):\n", " test_close(self.model.a.data, self.old_a - self.lr * self.grad_a)\n", " test_close(self.model.b.data, self.old_b - self.lr * self.grad_b)\n", " self.old_a,self.old_b = self.model.a.data.clone(),self.model.b.data.clone()\n", " test_close(self.model.a.grad.data, self.grad_a)\n", " test_close(self.model.b.grad.data, self.grad_b)\n", " \n", " def after_batch(self):\n", " for p in self.model.parameters(): test_eq(p.grad, tensor([0.]))" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#hide\n", "learn = synth_learner()\n", "b = learn.dls.one_batch()\n", "learn = synth_learner(cbs=TestOneBatch(*b, 42), lr=1e-2)\n", "#Remove train/eval\n", "learn.cbs = learn.cbs[1:]\n", "#Setup\n", "learn.loss,learn.training = tensor(0.),True\n", "learn.opt = SGD(learn.model.parameters(), lr=learn.lr)\n", "learn.model.train()\n", "batch_events = ['before_batch', 'after_pred', 'after_loss', 'before_backward', 'after_backward', 'after_step', 'after_batch']\n", "test_stdout(lambda: learn.one_batch(42, b), '\\n'.join(batch_events))\n", "test_stdout(lambda: learn.one_batch(42, b), '\\n'.join(batch_events)) #Check it works for a second batch" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

Learner.all_batches[source]

\n", "\n", "> Learner.all_batches()\n", "\n", "Train or evaluate `self.model` on all the batches of `self.dl`" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(Learner.all_batches)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#hide\n", "learn = synth_learner(n_train=5, cbs=VerboseCallback())\n", "learn.opt = SGD(learn.model.parameters(), lr=learn.lr)\n", "with redirect_stdout(io.StringIO()): \n", " learn(_before_epoch)\n", " learn.epoch,learn.dl = 0,learn.dls.train\n", " learn('before_epoch')\n", " learn('before_train')\n", "test_stdout(learn.all_batches, '\\n'.join(batch_events * 5))\n", "test_eq(learn.train_iter, 5)\n", "\n", "valid_events = ['before_batch', 'after_pred', 'after_loss', 'after_batch']\n", "with redirect_stdout(io.StringIO()): \n", " learn.dl = learn.dls.valid\n", " learn('before_validate')\n", "test_stdout(learn.all_batches, '\\n'.join(valid_events * 2))\n", "test_eq(learn.train_iter, 5)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#hide\n", "learn = synth_learner(n_train=5, cbs=VerboseCallback())\n", "test_stdout(lambda: learn(_before_epoch), 'before_fit\\nbefore_epoch')\n", "test_eq(learn.loss, tensor(0.))" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#hide\n", "learn.opt = SGD(learn.model.parameters(), lr=learn.lr)\n", "learn.epoch = 0\n", "test_stdout(lambda: learn._do_epoch_train(), '\\n'.join(['before_train'] + batch_events * 5 + ['after_train']))" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#hide\n", "test_stdout(learn._do_epoch_validate, '\\n'.join(['before_validate'] + valid_events * 2+ ['after_validate']))" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

Learner.create_opt[source]

\n", "\n", "> Learner.create_opt()\n", "\n", "Create an optimizer with default hyper-parameters" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(Learner.create_opt)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "This method is called internally to create the optimizer, the hyper-parameters are then adjusted by what you pass to `Learner.fit` or your particular schedulers (see `callback.schedule`)." ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "learn = synth_learner(n_train=5, cbs=VerboseCallback())\n", "assert learn.opt is None\n", "learn.create_opt()\n", "assert learn.opt is not None\n", "test_eq(learn.opt.hypers[0]['lr'], learn.lr)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "### Serializing" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

Learner.save[source]

\n", "\n", "> Learner.save(**`file`**, **`with_opt`**=*`True`*, **`pickle_protocol`**=*`2`*)\n", "\n", "Save model and optimizer state (if `with_opt`) to `self.path/self.model_dir/file`" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(Learner.save)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "`file` can be a `Path`, a `string` or a buffer. `pickle_protocol` is passed along to `torch.save`." ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

Learner.load[source]

\n", "\n", "> Learner.load(**`file`**, **`with_opt`**=*`None`*, **`device`**=*`None`*, **`strict`**=*`True`*)\n", "\n", "Load model and optimizer state (if `with_opt`) from `self.path/self.model_dir/file` using `device`" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(Learner.load)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "`file` can be a `Path`, a `string` or a buffer. Use `device` to load the model/optimizer state on a device different from the one it was saved." ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "with tempfile.TemporaryDirectory() as d:\n", " learn = synth_learner(path=d)\n", " learn.fit(1)\n", " \n", " #Test save created a file\n", " learn.save('tmp')\n", " assert (Path(d)/'models/tmp.pth').exists()\n", " \n", " #Test load did load the model\n", " learn1 = synth_learner(path=d)\n", " learn1 = learn1.load('tmp')\n", " test_eq(learn.model.a, learn1.model.a)\n", " test_eq(learn.model.b, learn1.model.b)\n", " test_eq(learn.opt.state_dict(), learn1.opt.state_dict())" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#hide\n", "#Test load works when the model is saved without opt\n", "with tempfile.TemporaryDirectory() as d:\n", " learn = synth_learner(path=d)\n", " learn.fit(1)\n", " learn.save('tmp', with_opt=False)\n", " learn1 = synth_learner(path=d)\n", " learn1 = learn1.load('tmp')\n", " test_eq(learn.model.a, learn1.model.a)\n", " test_eq(learn.model.b, learn1.model.b)\n", " test_ne(learn.opt.state_dict(), learn1.opt.state_dict())" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "### Callback handling" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "We only describe the basic functionality linked to `Callback`s here. To learn more about `Callback`s an how to write them, check the [callback.core](http://dev.fast.ai/callback.core) module documentation.\n", "\n", "Let's first see how the `Callback`s become attributes of `Learner`:" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#Test init with callbacks\n", "class TstCallback(Callback):\n", " def batch_begin(self): self.learn.a = self.a + 1\n", "\n", "tst_learn = synth_learner()\n", "test_eq(len(tst_learn.cbs), 1)\n", "assert isinstance(tst_learn.cbs[0], TrainEvalCallback)\n", "assert hasattr(tst_learn, ('train_eval'))\n", "\n", "tst_learn = synth_learner(cbs=TstCallback())\n", "test_eq(len(tst_learn.cbs), 2)\n", "assert isinstance(tst_learn.cbs[1], TstCallback)\n", "assert hasattr(tst_learn, ('tst'))" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "A name that becomes an existing attribute of the `Learner` will throw an exception (here add_cb is a method of `Learner`)." ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "class AddCbCallback(Callback): pass\n", "test_fail(lambda: synth_learner(cbs=AddCbCallback()))" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

Learner.__call__[source]

\n", "\n", "> Learner.__call__(**`event_name`**)\n", "\n", "Call `event_name` for all [`Callback`](/callback.core#Callback)s in `self.cbs`" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(Learner.__call__)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "This how the `Callback`s are called internally. For instance a `VerboseCallback` just prints the event names (can be useful for debugging):" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "name": "stdout", "output_type": "stream", "text": [ "after_fit\n" ] } ], "source": [ "learn = synth_learner(cbs=VerboseCallback())\n", "learn('after_fit')" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

Learner.add_cb[source]

\n", "\n", "> Learner.add_cb(**`cb`**)\n", "\n", "Add `cb` to the list of [`Callback`](/callback.core#Callback) and register `self` as their learner" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(Learner.add_cb)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "learn = synth_learner()\n", "learn.add_cb(TestTrainEvalCallback())\n", "test_eq(len(learn.cbs), 2)\n", "assert isinstance(learn.cbs[1], TestTrainEvalCallback)\n", "test_eq(learn.train_eval.learn, learn)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

Learner.add_cbs[source]

\n", "\n", "> Learner.add_cbs(**`cbs`**)\n", "\n", "Add `cbs` to the list of [`Callback`](/callback.core#Callback) and register `self` as their learner" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(Learner.add_cbs)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "learn.add_cbs([TestTrainEvalCallback(), TestTrainEvalCallback()])\n", "test_eq(len(learn.cbs), 4)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

Learner.added_cbs[source]

\n", "\n", "> Learner.added_cbs(**`cbs`**)\n", "\n", "Context manage that temporarily adds `cbs`" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(Learner.added_cbs)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "learn = synth_learner()\n", "test_eq(len(learn.cbs), 1)\n", "with learn.added_cbs(TestTrainEvalCallback()):\n", " test_eq(len(learn.cbs), 2)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

Learner.ordered_cbs[source]

\n", "\n", "> Learner.ordered_cbs(**`event`**)\n", "\n", "List of [`Callback`](/callback.core#Callback)s, in order, for an `event` in the training loop" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(Learner.ordered_cbs)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "By order, we mean using the internal ordering of the `Callback`s (see `callbcak.core` for more information on how it works)." ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/plain": [ "[TrainEvalCallback, TestTrainEvalCallback]" ] }, "execution_count": null, "metadata": {}, "output_type": "execute_result" } ], "source": [ "learn = synth_learner()\n", "learn.add_cb(TestTrainEvalCallback())\n", "learn.ordered_cbs('before_fit')" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

Learner.remove_cb[source]

\n", "\n", "> Learner.remove_cb(**`cb`**)\n", "\n", "Add `cb` from the list of [`Callback`](/callback.core#Callback) and deregister `self` as their learner" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(Learner.remove_cb)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "learn = synth_learner()\n", "learn.add_cb(TestTrainEvalCallback())\n", "cb = learn.cbs[1]\n", "learn.remove_cb(learn.cbs[1])\n", "test_eq(len(learn.cbs), 1)\n", "assert cb.learn is None\n", "assert not getattr(learn,'test_train_eval',None)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "`cb` can simply be the class of the `Callback` we want to remove (in which case all instances of that callback are removed)." ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "learn = synth_learner()\n", "learn.add_cbs([TestTrainEvalCallback(), TestTrainEvalCallback()])\n", "learn.remove_cb(TestTrainEvalCallback)\n", "test_eq(len(learn.cbs), 1)\n", "assert not getattr(learn,'test_train_eval',None)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

Learner.remove_cbs[source]

\n", "\n", "> Learner.remove_cbs(**`cbs`**)\n", "\n", "Remove `cbs` from the list of [`Callback`](/callback.core#Callback) and deregister `self` as their learner" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(Learner.remove_cbs)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "Elements of `cbs` can either be types of callbacks or actual callbacks of the `Learner`." ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "learn = synth_learner()\n", "learn.add_cbs([TestTrainEvalCallback() for _ in range(3)])\n", "cb = learn.cbs[1]\n", "learn.remove_cbs(learn.cbs[1:])\n", "test_eq(len(learn.cbs), 1)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

Learner.removed_cbs[source]

\n", "\n", "> Learner.removed_cbs(**`cbs`**)\n", "\n", "Context manage that temporarily removes `cbs`" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(Learner.removed_cbs)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "Elements of `cbs` can either be types of callbacks or actual callbacks of the `Learner`." ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "learn = synth_learner()\n", "learn.add_cb(TestTrainEvalCallback())\n", "with learn.removed_cbs(learn.cbs[1]):\n", " test_eq(len(learn.cbs), 1)\n", "test_eq(len(learn.cbs), 2)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

Learner.show_training_loop[source]

\n", "\n", "> Learner.show_training_loop()\n", "\n", "Show each step in the training loop" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(Learner.show_training_loop)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "At each step, callbacks are shown in order, which can help debugging." ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "name": "stdout", "output_type": "stream", "text": [ "Start Fit\n", " - before_fit : [TrainEvalCallback]\n", " Start Epoch Loop\n", " - before_epoch : []\n", " Start Train\n", " - before_train : [TrainEvalCallback]\n", " Start Batch Loop\n", " - before_batch : []\n", " - after_pred : []\n", " - after_loss : []\n", " - before_backward: []\n", " - after_backward : []\n", " - after_step : []\n", " - after_cancel_batch: []\n", " - after_batch : [TrainEvalCallback]\n", " End Batch Loop\n", " End Train\n", " - after_cancel_train: []\n", " - after_train : []\n", " Start Valid\n", " - before_validate: [TrainEvalCallback]\n", " Start Batch Loop\n", " - **CBs same as train batch**: []\n", " End Batch Loop\n", " End Valid\n", " - after_cancel_validate: []\n", " - after_validate : []\n", " End Epoch Loop\n", " - after_cancel_epoch: []\n", " - after_epoch : []\n", "End Fit\n", " - after_cancel_fit: []\n", " - after_fit : []\n" ] } ], "source": [ "learn = synth_learner()\n", "learn.show_training_loop()" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#export\n", "def _before_batch_cb(f, self):\n", " xb,yb = f(self, self.xb, self.yb)\n", " self.learn.xb,self.learn.yb = xb,yb" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#export\n", "def before_batch_cb(f):\n", " \"Shortcut for creating a Callback on the `before_batch` event, which takes and returns `xb,yb`\"\n", " return Callback(before_batch=partial(_before_batch_cb, f))" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "In order to change the data passed to your model, you will generally want to hook into the `before_batch` event, like so:" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "class TstCallback(Callback):\n", " def before_batch(self):\n", " self.learn.xb = self.xb + 1000\n", " self.learn.yb = self.yb - 1000" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "Since that is so common, we provide the `before_batch_cb` decorator to make it easier." ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "@before_batch_cb\n", "def cb(self, xb, yb): return xb+1000,yb-1000" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "### Control flow testing -" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#hide\n", "batch_events = ['before_batch', 'after_pred', 'after_loss', 'before_backward', 'after_backward', 'after_step', 'after_batch']\n", "batchv_events = ['before_batch', 'after_pred', 'after_loss', 'after_batch']\n", "train_events = ['before_train'] + batch_events + ['after_train']\n", "valid_events = ['before_validate'] + batchv_events + ['after_validate']\n", "epoch_events = ['before_epoch'] + train_events + valid_events + ['after_epoch']\n", "cycle_events = ['before_fit'] + epoch_events + ['after_fit']" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#hide\n", "learn = synth_learner(n_train=1, n_valid=1)\n", "test_stdout(lambda: learn.fit(1, cbs=VerboseCallback()), '\\n'.join(cycle_events))" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#hide\n", "class TestCancelCallback(VerboseCallback):\n", " def __init__(self, cancel_at=event.before_batch, exception=CancelBatchException, train=None):\n", " def _interrupt(): \n", " if train is None or train == self.training: raise exception()\n", " setattr(self, cancel_at, _interrupt)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#hide\n", "#test cancel batch\n", "for i,e in enumerate(batch_events[:-1]):\n", " be = batch_events[:i+1] + ['after_cancel_batch', 'after_batch']\n", " bev = be if i <3 else batchv_events\n", " cycle = cycle_events[:3] + be + ['after_train', 'before_validate'] + bev + cycle_events[-3:]\n", " test_stdout(lambda: learn.fit(1, cbs=TestCancelCallback(cancel_at=e)), '\\n'.join(cycle))\n", "\n", "#CancelBatchException not caught if thrown in any other event\n", "for e in cycle_events:\n", " if e not in batch_events[:-1]:\n", " with redirect_stdout(io.StringIO()):\n", " cb = TestCancelCallback(cancel_at=e)\n", " test_fail(lambda: learn.fit(1, cbs=cb))\n", " learn.remove_cb(cb) #Have to remove it manually" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#hide\n", "#test cancel train\n", "for i,e in enumerate(['before_train'] + batch_events):\n", " be = batch_events[:i] + (['after_batch'] if i >=1 and i < len(batch_events) else []) \n", " be += ['after_cancel_train', 'after_train']\n", " cycle = cycle_events[:3] + be + ['before_validate'] + batchv_events + cycle_events[-3:]\n", " test_stdout(lambda: learn.fit(1, cbs=TestCancelCallback(e, CancelTrainException, True)), '\\n'.join(cycle))\n", "\n", "#CancelTrainException not caught if thrown in any other event\n", "for e in cycle_events:\n", " if e not in ['before_train'] + batch_events[:-1]:\n", " with redirect_stdout(io.StringIO()):\n", " cb = TestCancelCallback(e, CancelTrainException)\n", " test_fail(lambda: learn.fit(1, cbs=cb))\n", " learn.remove_cb(cb) #Have to remove it manually " ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#hide\n", "#test cancel valid\n", "for i,e in enumerate(['before_validate'] + batchv_events):\n", " bev = batchv_events[:i] + (['after_batch'] if i >=1 and i < len(batchv_events) else []) + ['after_cancel_validate']\n", " cycle = cycle_events[:3] + batch_events + ['after_train', 'before_validate'] + bev + cycle_events[-3:]\n", " test_stdout(lambda: learn.fit(1, cbs=TestCancelCallback(e, CancelValidException, False)), '\\n'.join(cycle))\n", " \n", "#CancelValidException not caught if thrown in any other event\n", "for e in cycle_events:\n", " if e not in ['before_validate'] + batch_events[:3]:\n", " with redirect_stdout(io.StringIO()):\n", " cb = TestCancelCallback(e, CancelValidException)\n", " test_fail(lambda: learn.fit(1, cbs=cb))\n", " learn.remove_cb(cb) #Have to remove it manually " ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#hide\n", "#test cancel epoch\n", "#In train\n", "for i,e in enumerate(['before_train'] + batch_events):\n", " be = batch_events[:i] + (['after_batch'] if i >=1 and i=1 and i=1 and i=1 and iclass Metric[source]\n", "\n", "> Metric()\n", "\n", "Blueprint for defining a metric" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(Metric, title_level=3)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "Metrics can be simple averages (like accuracy) but sometimes their computation is a little bit more complex and can't be averaged over batches (like precision or recall), which is why we need a special class for them. For simple functions that can be computed as averages over batches, we can use the class `AvgMetric`, otherwise you'll need to implement the following methods.\n", "\n", "> Note: If your Metric has state depending on tensors, don't forget to store it on the CPU to avoid any potential memory leaks." ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

Metric.reset[source]

\n", "\n", "> Metric.reset()\n", "\n", "Reset inner state to prepare for new computation" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(Metric.reset)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

Metric.accumulate[source]

\n", "\n", "> Metric.accumulate(**`learn`**)\n", "\n", "Use `learn` to update the state with new results" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(Metric.accumulate)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

Metric.value[source]

\n", "\n", "The value of the metric" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(Metric.value, name='Metric.value')" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

Metric.name[source]

\n", "\n", "Name of the [`Metric`](/learner#Metric), camel-cased and with Metric removed" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(Metric.name, name='Metric.name')" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#export\n", "def _maybe_reduce(val):\n", " if num_distrib()>1:\n", " val = val.clone()\n", " torch.distributed.all_reduce(val, op=torch.distributed.ReduceOp.SUM)\n", " val /= num_distrib()\n", " return val" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#export\n", "class AvgMetric(Metric):\n", " \"Average the values of `func` taking into account potential different batch sizes\"\n", " def __init__(self, func): self.func = func\n", " def reset(self): self.total,self.count = 0.,0\n", " def accumulate(self, learn):\n", " bs = find_bs(learn.yb)\n", " self.total += to_detach(self.func(learn.pred, *learn.yb))*bs\n", " self.count += bs\n", " @property\n", " def value(self): return self.total/self.count if self.count != 0 else None\n", " @property\n", " def name(self): return self.func.func.__name__ if hasattr(self.func, 'func') else self.func.__name__" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

class AvgMetric[source]

\n", "\n", "> AvgMetric(**`func`**) :: [`Metric`](/learner#Metric)\n", "\n", "Average the values of `func` taking into account potential different batch sizes" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(AvgMetric, title_level=3)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "learn = synth_learner()\n", "tst = AvgMetric(lambda x,y: (x-y).abs().mean())\n", "t,u = torch.randn(100),torch.randn(100)\n", "tst.reset()\n", "for i in range(0,100,25): \n", " learn.pred,learn.yb = t[i:i+25],(u[i:i+25],)\n", " tst.accumulate(learn)\n", "test_close(tst.value, (t-u).abs().mean())" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#hide\n", "#With varying batch size\n", "tst.reset()\n", "splits = [0, 30, 50, 60, 100]\n", "for i in range(len(splits )-1): \n", " learn.pred,learn.yb = t[splits[i]:splits[i+1]],(u[splits[i]:splits[i+1]],)\n", " tst.accumulate(learn)\n", "test_close(tst.value, (t-u).abs().mean())" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#export\n", "class AvgLoss(Metric):\n", " \"Average the losses taking into account potential different batch sizes\"\n", " def reset(self): self.total,self.count = 0.,0\n", " def accumulate(self, learn):\n", " bs = find_bs(learn.yb)\n", " self.total += to_detach(learn.loss.mean())*bs\n", " self.count += bs\n", " @property\n", " def value(self): return self.total/self.count if self.count != 0 else None\n", " @property\n", " def name(self): return \"loss\"" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

class AvgLoss[source]

\n", "\n", "> AvgLoss() :: [`Metric`](/learner#Metric)\n", "\n", "Average the losses taking into account potential different batch sizes" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(AvgLoss, title_level=3)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "tst = AvgLoss()\n", "t = torch.randn(100)\n", "tst.reset()\n", "for i in range(0,100,25): \n", " learn.yb,learn.loss = t[i:i+25],t[i:i+25].mean()\n", " tst.accumulate(learn)\n", "test_close(tst.value, t.mean())" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#hide\n", "#With varying batch size\n", "tst.reset()\n", "splits = [0, 30, 50, 60, 100]\n", "for i in range(len(splits )-1): \n", " learn.yb,learn.loss = t[splits[i]:splits[i+1]],t[splits[i]:splits[i+1]].mean()\n", " tst.accumulate(learn)\n", "test_close(tst.value, t.mean())" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#export\n", "class AvgSmoothLoss(Metric):\n", " \"Smooth average of the losses (exponentially weighted with `beta`)\"\n", " def __init__(self, beta=0.98): self.beta = beta\n", " def reset(self): self.count,self.val = 0,tensor(0.)\n", " def accumulate(self, learn):\n", " self.count += 1\n", " self.val = torch.lerp(to_detach(learn.loss.mean(), gather=False), self.val, self.beta)\n", " @property\n", " def value(self): return self.val/(1-self.beta**self.count)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

class AvgSmoothLoss[source]

\n", "\n", "> AvgSmoothLoss(**`beta`**=*`0.98`*) :: [`Metric`](/learner#Metric)\n", "\n", "Smooth average of the losses (exponentially weighted with `beta`)" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(AvgSmoothLoss, title_level=3)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "tst = AvgSmoothLoss()\n", "t = torch.randn(100)\n", "tst.reset()\n", "val = tensor(0.)\n", "for i in range(4): \n", " learn.loss = t[i*25:(i+1)*25].mean()\n", " tst.accumulate(learn)\n", " val = val*0.98 + t[i*25:(i+1)*25].mean()*(1-0.98)\n", " test_close(val/(1-0.98**(i+1)), tst.value)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#export\n", "class ValueMetric(Metric):\n", " \"Use to include a pre-calculated metric value (for insance calculated in a `Callback`) and returned by `func`\"\n", " def __init__(self, func, metric_name=None): store_attr(self, 'func, metric_name')\n", "\n", " @property\n", " def value(self): return self.func()\n", "\n", " @property\n", " def name(self): return self.metric_name if self.metric_name else self.func.__name__" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

class ValueMetric[source]

\n", "\n", "> ValueMetric(**`func`**, **`metric_name`**=*`None`*) :: [`Metric`](/learner#Metric)\n", "\n", "Use to include a pre-calculated metric value (for insance calculated in a [`Callback`](/callback.core#Callback)) and returned by `func`" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(ValueMetric, title_level=3)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "def metric_value_fn(): return 5e-3\n", "\n", "vm = ValueMetric(metric_value_fn, 'custom_value_metric')\n", "test_eq(vm.value, 5e-3)\n", "test_eq(vm.name, 'custom_value_metric')\n", "\n", "vm = ValueMetric(metric_value_fn)\n", "test_eq(vm.name, 'metric_value_fn')" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "## Recorder --" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#export\n", "from fastprogress.fastprogress import format_time\n", "\n", "def _maybe_item(t):\n", " t = t.value\n", " return t.item() if isinstance(t, Tensor) and t.numel()==1 else t" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#export\n", "class Recorder(Callback):\n", " \"Callback that registers statistics (lr, loss and metrics) during training\"\n", " remove_on_fetch,run_after = True,TrainEvalCallback\n", "\n", " def __init__(self, add_time=True, train_metrics=False, valid_metrics=True, beta=0.98):\n", " store_attr(self, 'add_time,train_metrics,valid_metrics')\n", " self.loss,self.smooth_loss = AvgLoss(),AvgSmoothLoss(beta=beta)\n", "\n", " def before_fit(self):\n", " \"Prepare state for training\"\n", " self.lrs,self.iters,self.losses,self.values = [],[],[],[]\n", " names = self.metrics.attrgot('name')\n", " if self.train_metrics and self.valid_metrics:\n", " names = L('loss') + names\n", " names = names.map('train_{}') + names.map('valid_{}')\n", " elif self.valid_metrics: names = L('train_loss', 'valid_loss') + names\n", " else: names = L('train_loss') + names\n", " if self.add_time: names.append('time')\n", " self.metric_names = 'epoch'+names\n", " self.smooth_loss.reset()\n", "\n", " def after_batch(self):\n", " \"Update all metrics and records lr and smooth loss in training\"\n", " if len(self.yb) == 0: return\n", " mets = self._train_mets if self.training else self._valid_mets\n", " for met in mets: met.accumulate(self.learn)\n", " if not self.training: return\n", " self.lrs.append(self.opt.hypers[-1]['lr'])\n", " self.losses.append(self.smooth_loss.value)\n", " self.learn.smooth_loss = self.smooth_loss.value\n", "\n", " def before_epoch(self):\n", " \"Set timer if `self.add_time=True`\"\n", " self.cancel_train,self.cancel_valid = False,False\n", " if self.add_time: self.start_epoch = time.time()\n", " self.log = L(getattr(self, 'epoch', 0))\n", "\n", " def before_train (self): self._train_mets[1:].map(Self.reset())\n", " def before_validate(self): self._valid_mets.map(Self.reset())\n", " def after_train (self): self.log += self._train_mets.map(_maybe_item)\n", " def after_validate(self): self.log += self._valid_mets.map(_maybe_item)\n", " def after_cancel_train(self): self.cancel_train = True\n", " def after_cancel_validate(self): self.cancel_valid = True\n", "\n", " def after_epoch(self):\n", " \"Store and log the loss/metric values\"\n", " self.learn.final_record = self.log[1:].copy()\n", " self.values.append(self.learn.final_record)\n", " if self.add_time: self.log.append(format_time(time.time() - self.start_epoch))\n", " self.logger(self.log)\n", " self.iters.append(self.smooth_loss.count)\n", "\n", " @property\n", " def _train_mets(self):\n", " if getattr(self, 'cancel_train', False): return L()\n", " return L(self.smooth_loss) + (self.metrics if self.train_metrics else L())\n", "\n", " @property\n", " def _valid_mets(self):\n", " if getattr(self, 'cancel_valid', False): return L()\n", " return (L(self.loss) + self.metrics if self.valid_metrics else L())\n", "\n", " def plot_loss(self, skip_start=5, with_valid=True):\n", " plt.plot(list(range(skip_start, len(self.losses))), self.losses[skip_start:], label='train')\n", " if with_valid:\n", " idx = (np.array(self.iters)Recorder.before_fit[source]\n", "\n", "> Recorder.before_fit()\n", "\n", "Prepare state for training" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(Recorder.before_fit)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

Recorder.before_epoch[source]

\n", "\n", "> Recorder.before_epoch()\n", "\n", "Set timer if `self.add_time=True`" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(Recorder.before_epoch)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

Recorder.before_validate[source]

\n", "\n", "> Recorder.before_validate()\n", "\n", "Reset loss and metrics state" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(Recorder.before_validate)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

Recorder.after_batch[source]

\n", "\n", "> Recorder.after_batch()\n", "\n", "Update all metrics and records lr and smooth loss in training" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(Recorder.after_batch)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

Recorder.after_epoch[source]

\n", "\n", "> Recorder.after_epoch()\n", "\n", "Store and log the loss/metric values" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(Recorder.after_epoch)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "### Plotting tools" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

Recorder.plot_loss[source]

\n", "\n", "> Recorder.plot_loss(**`skip_start`**=*`5`*, **`with_valid`**=*`True`*)\n", "\n", "Plot the losses from `skip_start` and onward" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(Recorder.plot_loss)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "image/png": "iVBORw0KGgoAAAANSUhEUgAAAXQAAAD5CAYAAAA3Os7hAAAABHNCSVQICAgIfAhkiAAAAAlwSFlzAAALEgAACxIB0t1+/AAAADh0RVh0U29mdHdhcmUAbWF0cGxvdGxpYiB2ZXJzaW9uMy4yLjIsIGh0dHA6Ly9tYXRwbG90bGliLm9yZy+WH4yJAAAgAElEQVR4nO3deXiV5Z3/8feXLIQAIWSDQICwSEDWQGRVBGxV0KpVqrhUob8ZqlZErR2X6aKt09o67QhScXCjnQJWcasWaWuFAUHAsCMEZAkQIiQEwh5Ckvv3R45OjCE5kSTPyZPP67pycU6eO+d8cl/Jh5P7PIs55xARkcavmdcBRESkbqjQRUR8QoUuIuITKnQREZ9QoYuI+IQKXUTEJ8JrGmBmacCfK3yqG/BT59zTFcYYMB0YD5wCJjnn1lb3uAkJCS41NfXrZBYRabLWrFlzyDmXWNW2GgvdObcNGAhgZmHAfuDNSsPGARcEPoYCswL/nlNqaiqZmZk1hhcRkf9jZnvOta22Sy6XATudc5Uf8Frgj67cSiDWzJJr+dgiInIealvoE4H5VXy+I7Cvwv2cwOdERKSBBF3oZhYJXAO8VtXmKj73lXMKmNkUM8s0s8z8/PzgU4qISI1qXEOvYByw1jl3sIptOUCnCvdTgNzKg5xzs4HZABkZGV8p/LNnz5KTk0NRUVEtYjVOUVFRpKSkEBER4XUUEfGJ2hT6zVS93ALwF+AeM3uF8jdDjzrnPqttmJycHFq3bk1qairlO874k3OOgoICcnJy6Nq1q9dxRMQnglpyMbNo4JvAGxU+d6eZ3Rm4uxDYBewAngfu/jphioqKiI+P93WZA5gZ8fHxTeIvERFpOEG9QnfOnQLiK33uuQq3HfCDugjk9zL/XFP5PkWk4ehIURERn1ChV1BYWMizzz5b668bP348hYWF9ZBIRCR4KvQKzlXopaWl1X7dwoULiY2Nra9YIiJBqc1eLr738MMPs3PnTgYOHEhERAStWrUiOTmZ9evXs2XLFq677jr27dtHUVER06ZNY8qUKcD/ncbgxIkTjBs3josvvpgVK1bQsWNH3n77bVq0aOHxdyYiTUHIFvrj73zCltxjdfqYF3aI4Wff6nPO7U8++SSbN29m/fr1LFmyhKuuuorNmzd/sWvhSy+9RFxcHKdPn+aiiy7ihhtuID7+S+8V8+mnnzJ//nyef/55brzxRl5//XVuu+22Ov0+RESqErKFHgqGDBnypf3EZ8yYwZtvlp+XbN++fXz66adfKfSuXbsycOBAAAYPHkx2dnaD5RWRpi1kC726V9INpWXLll/cXrJkCe+//z4fffQR0dHRjB49usr9yJs3b/7F7bCwME6fPt0gWUVE9KZoBa1bt+b48eNVbjt69Cht27YlOjqarKwsVq5c2cDpRESqF7Kv0L0QHx/PyJEj6du3Ly1atKBdu3ZfbLvyyit57rnn6N+/P2lpaQwbNszDpCIiX2XlB3k2vIyMDFf5Ahdbt26ld+/enuTxQlP7fkXk/JnZGudcRlXbtOQiIuITKnQREZ9QoYuI+IQKXUTEJ1ToIiI+oUIXEfEJFfp5aNWqFQC5ublMmDChyjGjR4+m8u6ZIiL1QYVeBzp06MCCBQu8jiEiTZyOFK3goYceokuXLtx9d/klUR977DHMjKVLl3LkyBHOnj3LE088wbXXXvulr8vOzubqq69m8+bNnD59msmTJ7NlyxZ69+6tc7mISIMJ3UJ/72E4sKluH7N9Pxj35Dk3T5w4kfvuu++LQn/11VdZtGgR999/PzExMRw6dIhhw4ZxzTXXnPOaoLNmzSI6OpqNGzeyceNGBg0aVLffg4jIOYRuoXsgPT2dvLw8cnNzyc/Pp23btiQnJ3P//fezdOlSmjVrxv79+zl48CDt27ev8jGWLl3KvffeC0D//v3p379/Q34LItKEhW6hV/NKuj5NmDCBBQsWcODAASZOnMjcuXPJz89nzZo1REREkJqaWuVpcys616t3EZH6pDdFK5k4cSKvvPIKCxYsYMKECRw9epSkpCQiIiJYvHgxe/bsqfbrR40axdy5cwHYvHkzGzdubIjYIiLBFbqZxZrZAjPLMrOtZja80vY2ZvaOmW0ws0/MbHL9xK1/ffr04fjx43Ts2JHk5GRuvfVWMjMzycjIYO7cufTq1avar7/rrrs4ceIE/fv35ze/+Q1DhgxpoOQi0tQFdfpcM/sDsMw594KZRQLRzrnCCtsfBdo45x4ys0RgG9DeOVd8rsfU6XOb3vcrIuevutPn1riGbmYxwChgEkCgpCsXtQNaW/nicSvgMFByHplFRKSWglly6QbkAy+b2Toze8HMWlYaMxPoDeQCm4Bpzrmyuo0qIiLVCabQw4FBwCznXDpwEni40pgrgPVAB2AgMDPwyv5LzGyKmWWaWWZ+fn6VT+bVFZQaWlP5PkWk4QRT6DlAjnNuVeD+AsoLvqLJwBuu3A5gN/CVdw+dc7OdcxnOuYzExMSvPFFUVBQFBQW+LzvnHAUFBURFRXkdRUR8pMY1dOfcATPbZ2ZpzrltwGXAlkrD9gY+v8zM2gFpwK7ahklJSSEnJ4dzvXr3k6ioKFJSUryOISI+EuyBRVOBuYE9XHYBk83sTgDn3HPAL4A5ZrYJMOAh59yh2oaJiIiga9eutf0yEREhyEJ3zq0HKu8m81yF7bnA5XWYS0REaklHioqI+IQKXUTEJ1ToIiI+oUIXEfEJFbqIiE+o0EVEfEKFLiLiEyp0ERGfUKGLiPiECl1ExCdU6CIiPqFCFxHxCRW6iIhPqNBFRHxChS4i4hMqdBERn1Chi4j4hApdRMQnVOgiIj6hQhcR8QkVuoiIT6jQRUR8QoUuIuITQRW6mcWa2QIzyzKzrWY2vIoxo81svZl9Ymb/W/dRRUSkOuFBjpsOLHLOTTCzSCC64kYziwWeBa50zu01s6Q6zikiIjWosdDNLAYYBUwCcM4VA8WVht0CvOGc2xsYk1e3MUVEpCbBLLl0A/KBl81snZm9YGYtK43pCbQ1syVmtsbMbq/qgcxsipllmllmfn7+eUYXEZGKgin0cGAQMMs5lw6cBB6uYsxg4CrgCuAnZtaz8gM552Y75zKccxmJiYnnl1xERL4kmELPAXKcc6sC9xdQXvCVxyxyzp10zh0ClgID6i6miIjUpMZCd84dAPaZWVrgU5cBWyoNexu4xMzCzSwaGApsrdOkIiJSrWD3cpkKzA3s4bILmGxmdwI4555zzm01s0XARqAMeME5t7leEouISJXMOefJE2dkZLjMzExPnltEpLEyszXOuYyqtulIURERn1Chi4j4hApdRMQnVOgiIj6hQhcR8QkVuoiIT6jQRUR8QoUuIuITKnQREZ9QoYuI+IQKXUTEJ1ToIiI+oUIXEfEJFbqIiE+o0EVEfEKFLiLiEyp0ERGfUKGLiPiECl1ExCdU6CIiPqFCFxHxCRW6iIhPqNBFRHwiqEI3s1gzW2BmWWa21cyGn2PcRWZWamYT6jamiIjUJDzIcdOBRc65CWYWCURXHmBmYcCvgb/VYT4REQlSja/QzSwGGAW8COCcK3bOFVYxdCrwOpBXpwlFRCQowSy5dAPygZfNbJ2ZvWBmLSsOMLOOwLeB56p7IDObYmaZZpaZn5//tUOLiMhXBVPo4cAgYJZzLh04CTxcaczTwEPOudLqHsg5N9s5l+Gcy0hMTPxagUVEpGrBrKHnADnOuVWB+wv4aqFnAK+YGUACMN7MSpxzb9VZUhERqVaNhe6cO2Bm+8wszTm3DbgM2FJpTNfPb5vZHOBdlbmISMMKdi+XqcDcwB4uu4DJZnYngHOu2nVzERFpGEEVunNuPeXLKhVVWeTOuUnnmalauw+dZPr723nyhv5ERYTV51OJiDQqje5I0T0FJ3lrfS6Pv7Ol5sEiIk1Ioyv00WlJfP/SbsxfvZd3NuR6HUdEJGQ0ukIHePDyNNI7x/LIG5vYU3DS6zgiIiGhURZ6RFgznrk5nWYG98xbx5mSand/FxFpEhploQOktI3mqe8MYNP+o/z6vW1exxER8VyjLXSAK/q0Z9KIVF5avpt/bDnodRwREU816kIHeGR8L/p0iOHB1zawv/C013FERDzT6Au9eXgYM28ZRElpGdPmr6OktMzrSCIinmj0hQ7QNaElv7y+H5l7jvBf72/3Oo6IiCd8UegA1w7syE0ZnXh2yU6WbtepeUWk6fFNoQM8dk0feiS24oFX15N3vMjrOCIiDcpXhd4iMozf3zqIE2dKuO+V9ZSWOa8jiYg0GF8VOkDPdq15/Jo+rNhZwLOLd3gdR0Skwfiu0AFuzOjENQM68F/vb2f17sNexxERaRC+LHQz4z++3ZfOcdHcO38dh08Wex1JRKTe+bLQAVpHRTDzlkEcPlnMg69twDmtp4uIv/m20AH6dmzDo+N78UFWHi9+uNvrOCIi9crXhQ5wx4hULr+wHU++l8X6fYVexxERqTe+L3Qz46kJA2gXE8XU+Ws5evqs15FEROqF7wsdoE10BDNuTie3sIhH3tio9XQR8aUmUegAg7u05UdXpLFw0wHmrtrrdRwRkTrXZAodYMol3RjVM5Gfv7uFLbnHvI4jIlKngip0M4s1swVmlmVmW81seKXtt5rZxsDHCjMbUD9xz0+zZsbvbhxAbIsI7pm3lpNnSryOJCJSZ4J9hT4dWOSc6wUMALZW2r4buNQ51x/4BTC77iLWrYRWzXl64kB2F5zkJ29v9jqOiEidqbHQzSwGGAW8COCcK3bOfWn/P+fcCufckcDdlUBKXQetSyO6J3Dv2At4Y+1+FqzJ8TqOiEidCOYVejcgH3jZzNaZ2Qtm1rKa8f8PeK9O0tWjey+7gGHd4vjJW5vZkXfC6zgiIuctmEIPBwYBs5xz6cBJ4OGqBprZGMoL/aFzbJ9iZplmlpmf7+1FKMKaGdMnptMiMox75q2l6Gypp3lERM5XMIWeA+Q451YF7i+gvOC/xMz6Ay8A1zrnCqp6IOfcbOdchnMuIzEx8etmrjPtYqL47Y0DyDpwnJ+/u8XrOCIi56XGQnfOHQD2mVla4FOXAV9qPzPrDLwBfNc516gu6jkmLYnvj+rGvFV7eXdjrtdxRES+tvAgx00F5ppZJLALmGxmdwI4554DfgrEA8+aGUCJcy6jHvLWiwevSGN19mEeeX0T/TvG0jk+2utIIiK1Zl4dBp+RkeEyMzM9ee6q7Dt8iqtmLCM1oSUL7hxBZHiTOuZKRBoJM1tzrhfMaq2ATnHR/GZCfzbmHOXXi7K8jiMiUmsq9Aqu7JvMHcO78OKHu3l/y0Gv44iI1IoKvZJHxvfmwuQYHlywgdzC017HEREJmgq9kqiIMGbeks7ZkjLunb+OktIyryOJiARFhV6Fbomt+I9v9yNzzxGefv9Tr+OIiARFhX4O16V35MaMFH6/ZAcffnrI6zgiIjVSoVfjsWv60COxFff9eT15x4u8jiMiUi0VejWiI8OZecsgjhed5YE/b6CsTJeuE5HQpUKvQVr71jx+TR8+3HGIWf+70+s4IiLnpEIPwk0XdeJbAzrw279vY/Xuw17HERGpkgo9CGbGL7/dl05x0Ux7ZR1HThZ7HUlE5CtU6EFqHRXBzJsHcejEGR58bQNenQNHRORcVOi10C+lDY+O780/s/J48cPdXscREfkSFXotTRqRyjcvbMevF2WxYV9hzV8gItJAVOi1ZGY8NaE/Sa2juGf+Wo4VnfU6kogIoEL/WmKjI5lx80ByC4t45PVNWk8XkZCgQv+aBneJ44eX9+Svmz5j3uq9XscRaVDOObIOHGPWkp3c+N8f8d6mz7yOJAR/CTqpwp2juvPRzgIef2cLgzq3pXdyjNeRROrNqeISlu8oYPG2PJZk5ZF7tPx0GBfq5z5k6BJ05+nQiTOMm76MmKhw/nLPxbRsrv8jxT+yD53kg6w8Fm/LY9WuwxSXltEyMoyLL0hgTFoSo9OSaN8myuuYTUp1l6BT+5ynhFbNmX7TQG59cRU/ffsTfnvjAK8jiXxtZ0pKWb37MB9k5bFkWz67D50EoFtiS24f3oUxvZK4KDVO19wNUSr0OjCiRwJTx17AjH9+yoju8dwwOMXrSCJByy08zZJt+SzelsfyHYc4VVxKZHgzhneLZ9KIVEanJdIlvqXXMSUIKvQ6cu/YHqzcVcBP3t7MgE6x9Ehq5XUkkSqVlJaxdm8hi7flsTgrj6wDxwHoGNuC6wd1ZGyvJIZ3S6BFZJjHSaW2tIZehw4cLWLc9KW0i4nirR+MJCpCvxASGgpOnPniVfjS7fkcKyohvJmRkdqWMWlJjO2VRI+kVpiZ11GlBue9hm5mscALQF/AAd9zzn1UYbsB04HxwClgknNu7fkGb2zat4nidzcOZPKcj3nir1t44rp+XkeSJqqszLE592jgDc18NuYU4lz5ez5X9GnPmF5JXHxBAjFREV5HlToU7JLLdGCRc26CmUUC0ZW2jwMuCHwMBWYF/m1yxvRKYsqobsxeuovh3RK4qn+y15GkiTh6+iwffnqofLfCbfkcOnEGMxiQEst9l/VkbK8k+nSIoVkzvQr3qxoL3cxigFHAJADnXDFQ+fyx1wJ/dOXrNyvNLNbMkp1zTfJogwcvT2P17sM8/PpG+nVsQ+f4yv//iZw/5xzbD574Yi08c88RSsscbVpEMKpnImPSErm0ZyLxrZp7HVUaSDCv0LsB+cDLZjYAWANMc86drDCmI7Cvwv2cwOeaZKFHhjfjmZvTGT9jGVPnr+W1O0doNy+pE6eKS1jx+cE92/LZX3gagN7JMXx/VDfG9EoivVMs4WH6eWuKgin0cGAQMNU5t8rMpgMPAz+pMKaqv+G+8m6rmU0BpgB07ty59mkbkU5x0fzmhv7cNXctT/0ti3+/6kKvI0kjtafgJIuz8vhgWz4rdxVQXFJGdGQYF/dI4J6xPRidlkhymxZex5QQEEyh5wA5zrlVgfsLKC/0ymM6VbifAuRWfiDn3GxgNpTv5VLrtI3MuH7JfHdYF55ftpth3eK5rHc7ryNJI3CmpJSPdx/5Yill1+cH9yS05LahXRjbK4mLuralebj2opIvq7HQnXMHzGyfmaU557YBlwFbKg37C3CPmb1C+ZuhR5vq+nll/35VbzL3HOGHr23gvWmX6JWUVOmzo4GDe7LKD+45GTi4Z1i3eL47vAtj0pJITdDBPVK9YPdymQrMDezhsguYbGZ3AjjnngMWUr7L4g7Kd1ucXA9ZG6WoiDB+f0s6Vz/zIdPmr2fevw7V+qZQUlrG+n2FX+xWuPWzYwB0aBPFdekdGZOWxIge8URH6tg/CZ4OLGogb67L4f4/b2Dq2B788PI0r+OIBwpOnGHpp/l8kJXP0u35HD19lrBmxuAubRnbK4kxaUn0bKeDe6R6OjlXCPh2egordhQwc/EOhnWLZ2SPBK8jST0rK3N8knusfC18Wx7r931+cE8k3+jdjrGBg3vatNDBPVI3VOgN6PFr+7BuXyHTXlnPe9MuIbG19g/2m2NFgYN7svJYsj2f/OPlB/f0T4ll2mUXMCYtiX4d2+jgHqkXKvQGFB0Zzsxb0rl25nIeeHU9f5g8RL/YjZxzjh15J744Z3hm9hFKyhwxUeGBg3uSuDQtkQQd3CMNQIXewHq1j+Gxa/rwyBubmPW/O/nBmB5eR5JaOl1cyke7DpWXeNb/HdzTq31r/nVUN8akJTGosw7ukYanQvfAxIs6sXzHIX73j+0M7RpHRmqc15GkBnsLTrF4Wx4fZOXxUeDgnhYRYYzskcDdY7ozJi2JDrHaJVW8pUL3gJnxq+v7sWn/Ue6dv46/3nsJbVtGeh1LKiguKePj7MMsDiyl7MwvP7gnNT6aW4d2ZkxaEkO6xukUyRJSVOgeaR0VwTM3p3PDrBX8aMFGnr99sHZX85hzjg93HGLeqr0s3Z5ffnBPWDOGdovj1qHll1/rqoN7JISp0D3UPyWWR8b15ufvbuHl5dl87+KuXkdqkk4Vl/DG2v3MWZHNjrwTxLeM5JqB5VfuGdE9Xhf+lkZDP6kemzwylRU7C/jVe1vJSG1L/5RYryM1GfsOn+J/Vu7hldV7OVZUQt+OMfz2OwO4ekCyzpMijZKOFA0BhaeKGT99GeFhzXj33ot1FZl65Jxj5a7DzFmxm39sOYiZcWWf9kwemcrgLm217CUhT0eKhrjY6Ehm3JzOTbNX8sgbm5h5c7qKpY4VnS3l7fX7eXl5NlkHjhMbHcH3L+3Od4d10d4p4hsq9BCRkRrHA9/syVN/28bI7gncMtTf54tvKJ8dPc3/fLSH+av3cuTUWXq1b82vb+jHtQM7ag8V8R0Vegi569LurNxVwOPvfMKgLrH0ah/jdaRGyTnH2r1HeGl5Nos2H6DMOb7Zux2TRqYyvFu8/voR39IaeojJP36G8TOWERMVzjtTL9bpU2vhTEkpf934GS8vz2bT/qO0jgpn4kWduH14Kp3idF1X8QetoTciia2b8/RNA7ntxVX87O1PeOo7A7yOFPLyjhcxd+Ve5q7ay6ETZ+ie2JJfXNeX69M7apdDaVL00x6CRvZIYOqYHsz4YAfDu8dz/aAUryOFpA37CpmzIpt3N+ZyttQxtlcSk0akcnGPBJ30TJokFXqIuveyC1i56zA/fmszAzrF0j2xldeRQsLZ0jLe23yAOct3s3ZvIa2ah3Pr0C7cMSJVR3FKk6c19BD22dHTjJ++jPZtWvDm3SOa9F4ZBSfOMH/1Xv5n5R4OHjtDanw0d4xIZcLgFFprv31pQrSG3kglt2nBb28cwPfmZPIff93KL67r63WkBvdJ7lHmLM/m7Q25FJeUcckFCfzq+n6M7pmkZRWRSlToIW5sr3b86yVdeX7ZbkZ0j2dcv2SvI9W7ktIy3t96kJeWZ7N692FaRITxncEpTBqRygXtWnsdTyRkqdAbgR9d0YvV2Uf4t9c30rdjG9/ugld4qpg/f7yPP360h/2Fp+kY24JHx/fipozOtInWsopITbSG3kjsO3yK8TOW0S2xFa99fziR4f65Gs72g8eZsyKbN9bmUHS2jGHd4pg8sivf6N2OMC2riHyJ1tB9oFNcNL++oT93z13Lf/59G4+O7+11pPNSWuZYnJXHyyt2s3xHAc3Dm3HdwI7cMSKVCzvoCFmRryOoQjezbOA4UAqUVP7fwczaAH8COgce8z+dcy/XbVQZ3y+Z24Z1ZvbSXQzrFsfYXu28jlRrx4rO8lpmDn9Ykc3ew6doHxPFj65I4+YhnYnTVZtEzkttXqGPcc4dOse2HwBbnHPfMrNEYJuZzXXOFZ9/RKnox1ddSGb2EX746gYWTruE5DaN40yBu/JP8IcV2SxYk8PJ4lIGd2nLv12ZxhV92hOhiymL1Im6WnJxQGsrP+tRK+AwUFJHjy0VREWE8ftbB/GtZz5k2ivrmfcvQ0P26vJlZY6ln+YzZ0U2S7blExnWjKsHJDNpRKou5CFSD4ItdAf83cwc8N/OudmVts8E/gLkAq2Bm5xzZXUXUyrqntiKJ67rywOvbmDGBzt44Js9vY70JSfPlPD62hzmrMhmV/5JEls35/5v9OSWoZ1JbN3c63givhVsoY90zuWaWRLwDzPLcs4trbD9CmA9MBboHhizzDl3rOKDmNkUYApA58463/f5uH5QCit2FvDMB58yrGscI3okeB2JvQWn+ONH2fw5cx/Hi0oYkNKGp28ayPh+yb7aK0ckVNV6t0Uzeww44Zz7zwqf+yvwpHNuWeD+B8DDzrnV53oc7bZ4/k4Vl/CtZz7kWFEJC++9xJNXv845PtpZwEvLs/ln1kHCzBjXL5nJI1NJ7xSrc4+L1LHz2m3RzFoCzZxzxwO3Lwd+XmnYXuAyYJmZtQPSgF3nF1tqEh0ZzsxbBnHd75fzwKvr+cPkIQ12OPzp4lLeWr+fOcuz2XbwOHEtI/nB6B7cNqwL7dtENUgGEfmyYJZc2gFvBl5phQPznHOLzOxOAOfcc8AvgDlmtgkw4KFq9oiROtQ7OYaffutC/v3NzTy3dCd3j+5Rr8+3v7D8km6vfLyXwlNn6Z0cw28m9OeaAR2a9MnDREJBjYXunNsFfOUqC4Ei//x2LuWv3MUDtwzpzIqdBfz279sZ2jWOwV3i6vTxnXN8nH2EOSt287dPDuKc44o+7Zk8sisXpbbVsopIiNCRoj5gZvzq+n5syjnK1HnrWDjtEmKjz/8gnaKzpbyzIZc5K7L5JPcYbVpE8C+XdOW7w7qQ0taf55MRacxU6D4RExXBzFvSuWHWCh58bSPP3z74a79yPnisiD+t3MO8VXspOFlMz3at+OW3+3Fdegdd41QkhOm300f6p8Ty8Lje/OLdLcxZkc3kkV1r9fXr9h7h5eXZLNz0GaXOcVmvJCaP7MqI7vFaVhFpBFToPvO9kal8tPMQv1y4lYwucfRLaVPt+OKSMhZu+oyXV2SzYV8hrZuHc8eIVG4f3oUu8bqkm0hjotPn+tCRk8WMn7GMyPBmvDv14iov0ZZ//AzzVu3lT6v2kH/8DN0SWjJpZCrXD0qhVXP9Py8SqnT63CambctIZtyczsTZK3n0zc3MmDjwiyWTzfuP8vLybN7ZkEtxaRmX9kxk8oRURl2QqEu6iTRyKnSfuig1jge+2ZOn/raNoV3jaBsdyZwVu/k4+wjRkWFMHNKJO0ak0j2xlddRRaSOqNB97K5Lu7NyVwE/fmszAJ3iWvDjq3pz40WdiKliGUZEGjcVuo81a2b87saBTP/ndi7tmcTYXkm6pJuIj6nQfS6xdXOeuK6f1zFEpAHonKYiIj6hQhcR8QkVuoiIT6jQRUR8QoUuIuITKnQREZ9QoYuI+IQKXUTEJzw726KZ5QN7vuaXJwCheM3SUM0FoZtNuWpHuWrHj7m6OOcSq9rgWaGfDzPLPNfpI70UqrkgdLMpV+0oV+00tVxachER8QkVuoiITzTWQp/tdYBzCNVcELrZlKt2lKt2mlSuRrmGLiIiX9VYX6GLiEglIV3oZvaSmeWZ2eZzbDczm2FmO8xso5kNCpFco83sqJmtD3z8tAEydTKzxWa21cw+MbNpVYxp8PkKMmHDHLkAAANCSURBVJcX8xVlZqvNbEMg1+NVjPFivoLJ1eDzVeG5w8xsnZm9W8U2T34fg8jl5Xxlm9mmwPNmVrG9bufMOReyH8AoYBCw+RzbxwPvAQYMA1aFSK7RwLsNPFfJwKDA7dbAduBCr+cryFxezJcBrQK3I4BVwLAQmK9gcjX4fFV47geAeVU9v1e/j0Hk8nK+soGEarbX6ZyF9Ct059xS4HA1Q64F/ujKrQRizSw5BHI1OOfcZ865tYHbx4GtQMdKwxp8voLM1eACc3AicDci8FH5DSUv5iuYXJ4wsxTgKuCFcwzx5PcxiFyhrE7nLKQLPQgdgX0V7ucQAmURMDzwZ/N7ZtanIZ/YzFKBdMpf3VXk6XxVkws8mK/An+nrgTzgH865kJivIHKBNz9fTwP/BpSdY7tXP1815QLvfh8d8HczW2NmU6rYXqdz1tgLvaorHofCq5m1lB+eOwB4BniroZ7YzFoBrwP3OeeOVd5cxZc0yHzVkMuT+XLOlTrnBgIpwBAz61tpiCfzFUSuBp8vM7sayHPOraluWBWfq9f5CjKXZ7+PwEjn3CBgHPADMxtVaXudzlljL/QcoFOF+ylArkdZvuCcO/b5n83OuYVAhJkl1PfzmlkE5aU51zn3RhVDPJmvmnJ5NV8Vnr8QWAJcWWmTpz9f58rl0XyNBK4xs2zgFWCsmf2p0hgv5qvGXF7+fDnncgP/5gFvAkMqDanTOWvshf4X4PbAO8XDgKPOuc+8DmVm7c3MAreHUD7PBfX8nAa8CGx1zv3uHMMafL6CyeXRfCWaWWzgdgvgG0BWpWFezFeNubyYL+fcI865FOdcKjAR+MA5d1ulYQ0+X8Hk8mK+As/V0sxaf34buByovGdcnc5Z+NdO2wDMbD7l71AnmFkO8DPK3yTCOfccsJDyd4l3AKeAySGSawJwl5mVAKeBiS7wlnY9Ggl8F9gUWH8FeBToXCGXF/MVTC4v5isZ+IOZhVH+C/6qc+5dM7uzQi4v5iuYXF7MV5VCYL6CyeXVfLUD3gz8XxIOzHPOLarPOdORoiIiPtHYl1xERCRAhS4i4hMqdBERn1Chi4j4hApdRMQnVOgiIj6hQhcR8QkVuoiIT/x/9koVMB8zJAoAAAAASUVORK5CYII=\n", "text/plain": [ "
" ] }, "metadata": { "needs_background": "light" }, "output_type": "display_data" } ], "source": [ "#hide\n", "learn.recorder.plot_loss(skip_start=1)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "## Inference functions" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

Learner.validate[source]

\n", "\n", "> Learner.validate(**`ds_idx`**=*`1`*, **`dl`**=*`None`*, **`cbs`**=*`None`*)\n", "\n", "Validate on `dl` with potential new `cbs`." ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(Learner.validate)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#Test result\n", "learn = synth_learner(n_train=5, metrics=tst_metric)\n", "res = learn.validate()\n", "test_eq(res[0], res[1])\n", "x,y = learn.dls.valid_ds.tensors\n", "test_close(res[0], F.mse_loss(learn.model(x), y))" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#hide\n", "#Test other dl\n", "res = learn.validate(dl=learn.dls.train)\n", "test_eq(res[0], res[1])\n", "x,y = learn.dls.train_ds.tensors\n", "test_close(res[0], F.mse_loss(learn.model(x), y))\n", "\n", "#Test additional callback is executed.\n", "cycle = cycle_events[:2] + ['before_validate'] + batchv_events * 2 + cycle_events[-3:]\n", "test_stdout(lambda: learn.validate(cbs=VerboseCallback()), '\\n'.join(cycle))" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

Learner.get_preds[source]

\n", "\n", "> Learner.get_preds(**`ds_idx`**=*`1`*, **`dl`**=*`None`*, **`with_input`**=*`False`*, **`with_decoded`**=*`False`*, **`with_loss`**=*`False`*, **`act`**=*`None`*, **`inner`**=*`False`*, **`reorder`**=*`True`*, **`cbs`**=*`None`*, **`save_preds`**=*`None`*, **`save_targs`**=*`None`*, **`concat_dim`**=*`0`*)\n", "\n", "Get the predictions and targets on the `ds_idx`-th dbunchset or `dl`, optionally `with_input` and `with_loss`" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(Learner.get_preds)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "`with_decoded` will also return the decoded predictions using the `decodes` function of the loss function (if it exists). For instance, fastai's `CrossEntropyFlat` takes the argmax or predictions in its decodes. \n", "\n", "Depending on the `loss_func` attribute of `Learner`, an activation function will be picked automatically so that the predictions make sense. For instance if the loss is a case of cross-entropy, a softmax will be applied, or if the loss is binary cross entropy with logits, a sigmoid will be applied. If you want to make sure a certain activation function is applied, you can pass it with `act`.\n", "\n", "`save_preds` and `save_targs` should be used when your predictions are too big to fit all in memory. Give a `Path` object that points to a folder where the predictions and targets will be saved.\n", "\n", "`concat_dim` is the batch dimension, where all the tensors will be concatenated.\n", "\n", "`inner` is an internal attribute that tells `get_preds` it's called internally, inside another training loop, to avoid recursion errors." ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "> Note: If you want to use the option `with_loss=True` on a custom loss function, make sure you have implemented a `reduction` attribute that supports 'none' " ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#Test result\n", "learn = synth_learner(n_train=5, metrics=tst_metric)\n", "preds,targs = learn.get_preds()\n", "x,y = learn.dls.valid_ds.tensors\n", "test_eq(targs, y)\n", "test_close(preds, learn.model(x))\n", "\n", "preds,targs = learn.get_preds(act = torch.sigmoid)\n", "test_eq(targs, y)\n", "test_close(preds, torch.sigmoid(learn.model(x)))" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#hide\n", "#Test get_preds work with ds not evenly divisible by bs\n", "learn = synth_learner(n_train=2.5, metrics=tst_metric)\n", "preds,targs = learn.get_preds(ds_idx=0)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#hide\n", "#Test other dataset\n", "x = torch.randn(16*5)\n", "y = 2*x + 3 + 0.1*torch.randn(16*5)\n", "dl = TfmdDL(TensorDataset(x, y), bs=16)\n", "preds,targs = learn.get_preds(dl=dl)\n", "test_eq(targs, y)\n", "test_close(preds, learn.model(x))\n", "\n", "#Test with loss\n", "preds,targs,losses = learn.get_preds(dl=dl, with_loss=True)\n", "test_eq(targs, y)\n", "test_close(preds, learn.model(x))\n", "test_close(losses, F.mse_loss(preds, targs, reduction='none'))\n", "\n", "#Test with inputs\n", "inps,preds,targs = learn.get_preds(dl=dl, with_input=True)\n", "test_eq(inps,x)\n", "test_eq(targs, y)\n", "test_close(preds, learn.model(x))" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#hide\n", "#Test with no target\n", "learn = synth_learner(n_train=5)\n", "x = torch.randn(16*5)\n", "dl = TfmdDL(TensorDataset(x), bs=16)\n", "preds,targs = learn.get_preds(dl=dl)\n", "assert targs is None" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#hide\n", "#Test with targets that are tuples\n", "def _fake_loss(x,y,z,reduction=None): return F.mse_loss(x,y)\n", "\n", "learn = synth_learner(n_train=5)\n", "x = torch.randn(16*5)\n", "y = 2*x + 3 + 0.1*torch.randn(16*5)\n", "learn.dls.n_inp=1\n", "learn.loss_func = _fake_loss\n", "dl = TfmdDL(TensorDataset(x, y, y), bs=16)\n", "preds,targs = learn.get_preds(dl=dl)\n", "test_eq(targs, [y,y])" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#hide\n", "#Test with inputs that are tuples\n", "class _TupleModel(Module):\n", " def __init__(self, model): self.model=model\n", " def forward(self, x1, x2): return self.model(x1)\n", "\n", "learn = synth_learner(n_train=5)\n", "#learn.dls.n_inp=2\n", "x = torch.randn(16*5)\n", "y = 2*x + 3 + 0.1*torch.randn(16*5)\n", "learn.model = _TupleModel(learn.model)\n", "learn.dls = DataLoaders(TfmdDL(TensorDataset(x, x, y), bs=16),TfmdDL(TensorDataset(x, x, y), bs=16))\n", "inps,preds,targs = learn.get_preds(ds_idx=0, with_input=True)\n", "test_eq(inps, [x,x])\n", "t = learn.get_preds(ds_idx=0, with_input=True)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#hide\n", "#Test auto activation function is picked\n", "learn = synth_learner(n_train=5)\n", "learn.loss_func = BCEWithLogitsLossFlat()\n", "x = torch.randn(16*5)\n", "y = 2*x + 3 + 0.1*torch.randn(16*5)\n", "dl = TfmdDL(TensorDataset(x, y), bs=16)\n", "preds,targs = learn.get_preds(dl=dl)\n", "test_close(preds, torch.sigmoid(learn.model(x)))" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#hide\n", "#Test reorder is done\n", "learn = synth_learner(n_train=5)\n", "x = torch.randn(16*5)\n", "y = 2*x + 3 + 0.1*torch.randn(16*5)\n", "dl = TfmdDL(TensorDataset(x, y), bs=16, shuffle=True)\n", "preds,targs = learn.get_preds(dl=dl)\n", "test_eq(targs, y)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#hide\n", "inps,preds,targs = learn.get_preds(ds_idx=0, with_input=True)\n", "tst = learn.get_preds(ds_idx=0, with_input=True, with_decoded=True)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

Learner.predict[source]

\n", "\n", "> Learner.predict(**`item`**, **`rm_type_tfms`**=*`None`*, **`with_input`**=*`False`*)\n", "\n", "Prediction on `item`, fully decoded, loss function decoded and probabilities" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(Learner.predict)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "It returns a tuple of three elements with, in reverse order,\n", "- the prediction from the model, potentially passed through the activation of the loss function (if it has one)\n", "- the decoded prediction, using the poential `decodes` method from it\n", "- the fully decoded prediction, using the transforms used to buil the `Datasets`/`DataLoaders`" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "`rm_type_tfms` is a deprecated argument that should not be used and will be removed in a future version. `with_input` will add the decoded inputs to the result." ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "class _FakeLossFunc(Module):\n", " reduction = 'none'\n", " def forward(self, x, y): return F.mse_loss(x,y)\n", " def activation(self, x): return x+1\n", " def decodes(self, x): return 2*x\n", "\n", "class _Add1(Transform):\n", " def encodes(self, x): return x+1\n", " def decodes(self, x): return x-1\n", " \n", "learn = synth_learner(n_train=5)\n", "dl = TfmdDL(Datasets(torch.arange(50), tfms = [L(), [_Add1()]]))\n", "learn.dls = DataLoaders(dl, dl)\n", "learn.loss_func = _FakeLossFunc()\n", "\n", "inp = tensor([2.])\n", "out = learn.model(inp).detach()+1 #applying model + activation\n", "dec = 2*out #decodes from loss function\n", "full_dec = dec-1 #decodes from _Add1\n", "test_eq(learn.predict(inp), [full_dec,dec,out])\n", "test_eq(learn.predict(inp, with_input=True), [inp,full_dec,dec,out])" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

Learner.show_results[source]

\n", "\n", "> Learner.show_results(**`ds_idx`**=*`1`*, **`dl`**=*`None`*, **`max_n`**=*`9`*, **`shuffle`**=*`True`*, **\\*\\*`kwargs`**)\n", "\n", "Show some predictions on `ds_idx`-th dataset or `dl`" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(Learner.show_results)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "Will show `max_n` samples (unless the batch size of `ds_idx` or `dl` is less than `max_n`, in which case it will show as many samples) and `shuffle` the data unless you pass `false` to that flag. `kwargs` are application-dependant.\n", "\n", "We can't show an example on our synthetic `Learner`, but check all the beginners tutorials which will show you how that method works accross applications." ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "The last functions in this section are used internally for inference, but should be less useful to you." ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

Learner.no_logging[source]

\n", "\n", "> Learner.no_logging()\n", "\n", "Context manager to temporarily remove `logger`" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(Learner.no_logging)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "learn = synth_learner(n_train=5, metrics=tst_metric)\n", "with learn.no_logging():\n", " test_stdout(lambda: learn.fit(1), '')\n", "test_eq(learn.logger, print)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

Learner.loss_not_reduced[source]

\n", "\n", "> Learner.loss_not_reduced()\n", "\n", "A context manager to evaluate `loss_func` with reduction set to none." ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(Learner.loss_not_reduced)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "This requires your loss function to either have a `reduction` attribute or a `reduction` argument (like all fastai and PyTorch loss functions)." ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#hide\n", "test_eq(learn.loss_func.reduction, 'mean')\n", "with learn.loss_not_reduced():\n", " test_eq(learn.loss_func.reduction, 'none')\n", " x,y = learn.dls.one_batch()\n", " p = learn.model(x)\n", " losses = learn.loss_func(p, y)\n", " test_eq(losses.shape, y.shape)\n", " test_eq(losses, F.mse_loss(p,y, reduction='none'))\n", "test_eq(learn.loss_func.reduction, 'mean')" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "## Transfer learning" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#export\n", "@patch\n", "def freeze_to(self:Learner, n):\n", " if self.opt is None: self.create_opt()\n", " self.opt.freeze_to(n)\n", " self.opt.clear_state()\n", "\n", "@patch\n", "def freeze(self:Learner): self.freeze_to(-1)\n", "\n", "@patch\n", "def unfreeze(self:Learner): self.freeze_to(0)\n", "\n", "add_docs(Learner,\n", " freeze_to=\"Freeze parameter groups up to `n`\",\n", " freeze=\"Freeze up to last parameter group\",\n", " unfreeze=\"Unfreeze the entire model\")" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "name": "stdout", "output_type": "stream", "text": [ "(#4) [0,16.64426040649414,16.32911491394043,'00:00']\n" ] } ], "source": [ "#hide\n", "class _TstModel(nn.Module):\n", " def __init__(self):\n", " super().__init__()\n", " self.a,self.b = nn.Parameter(torch.randn(1)),nn.Parameter(torch.randn(1))\n", " self.tst = nn.Sequential(nn.Linear(4,5), nn.BatchNorm1d(3))\n", " self.tst[0].bias.data,self.tst[1].bias.data = torch.randn(5),torch.randn(3) \n", " def forward(self, x): return x * self.a + self.b\n", " \n", "class _PutGrad(Callback):\n", " def after_backward(self):\n", " for p in self.learn.model.tst.parameters():\n", " if p.requires_grad: p.grad = torch.ones_like(p.data)\n", "\n", "def _splitter(m): return [list(m.tst[0].parameters()), list(m.tst[1].parameters()), [m.a,m.b]]\n", " \n", "learn = synth_learner(n_train=5, opt_func = partial(SGD), cbs=_PutGrad, splitter=_splitter, lr=1e-2)\n", "learn.model = _TstModel()\n", "learn.freeze()\n", "init = [p.clone() for p in learn.model.tst.parameters()]\n", "learn.fit(1, wd=0.)\n", "end = list(learn.model.tst.parameters())\n", "#linear was not trained\n", "for i in [0,1]: test_close(end[i],init[i])\n", "#bn was trained even frozen since `train_bn=True` by default\n", "for i in [2,3]: test_close(end[i]-init[i], -0.05 * torch.ones_like(end[i]))" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "name": "stdout", "output_type": "stream", "text": [ "(#4) [0,10.765480041503906,10.154903411865234,'00:00']\n", "(#4) [0,8.656882286071777,8.24548625946045,'00:00']\n", "(#4) [0,7.037893772125244,6.696390151977539,'00:00']\n" ] } ], "source": [ "#hide\n", "learn = synth_learner(n_train=5, opt_func = partial(SGD), cbs=_PutGrad, splitter=_splitter, train_bn=False, lr=1e-2)\n", "learn.model = _TstModel()\n", "learn.freeze()\n", "init = [p.clone() for p in learn.model.tst.parameters()]\n", "learn.fit(1, wd=0.)\n", "end = list(learn.model.tst.parameters())\n", "#linear and bn were not trained\n", "for i in range(4): test_close(end[i],init[i])\n", "\n", "learn.freeze_to(-2)\n", "init = [p.clone() for p in learn.model.tst.parameters()]\n", "learn.fit(1, wd=0.)\n", "end = list(learn.model.tst.parameters())\n", "#linear was not trained\n", "for i in [0,1]: test_close(end[i],init[i])\n", "#bn was trained \n", "for i in [2,3]: test_close(end[i]-init[i], -0.05 * torch.ones_like(end[i]))\n", " \n", "learn.unfreeze()\n", "init = [p.clone() for p in learn.model.tst.parameters()]\n", "learn.fit(1, wd=0.)\n", "end = list(learn.model.tst.parameters())\n", "#linear and bn were trained\n", "for i in range(4): test_close(end[i]-init[i], -0.05 * torch.ones_like(end[i]), 1e-3)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "### Exporting a `Learner`" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#export\n", "@patch\n", "def export(self:Learner, fname='export.pkl', pickle_protocol=2):\n", " \"Export the content of `self` without the items and the optimizer state for inference\"\n", " if rank_distrib(): return # don't export if child proc\n", " self._end_cleanup()\n", " old_dbunch = self.dls\n", " self.dls = self.dls.new_empty()\n", " state = self.opt.state_dict() if self.opt is not None else None\n", " self.opt = None\n", " with warnings.catch_warnings():\n", " #To avoid the warning that come from PyTorch about model not being checked\n", " warnings.simplefilter(\"ignore\")\n", " torch.save(self, self.path/fname, pickle_protocol=pickle_protocol)\n", " self.create_opt()\n", " if state is not None: self.opt.load_state_dict(state)\n", " self.dls = old_dbunch" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "The `Learner` is saved in `self.path/fname`, using `pickle_protocol`. Note that serialization in Python saves the names of functions, not the code itself. Therefore, any custom code you have for models, data transformation, loss function etc... should be put in a module that you will import in your training environment before exporting, and in your deployment environment before loading it." ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#export\n", "def load_learner(fname, cpu=True):\n", " \"Load a `Learner` object in `fname`, optionally putting it on the `cpu`\"\n", " distrib_barrier()\n", " res = torch.load(fname, map_location='cpu' if cpu else None)\n", " if hasattr(res, 'to_fp32'): res = res.to_fp32()\n", " if cpu: res.dls.cpu()\n", " return res" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "> Warning: `load_learner` requires all your custom code be in the exact same place as when exporting your `Learner` (the main script, or the module you imported it from)." ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "## TTA" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#export\n", "def _tta(self:Learner, ds_idx=1, dl=None, n=4, item_tfms=None, batch_tfms=None, beta=0.25, use_max=False):\n", " with dl.dataset.set_split_idx(0), self.no_mbar():\n", " if hasattr(self,'progress'): self.progress.mbar = master_bar(list(range(n)))\n", " aug_preds = []\n", " for i in self.progress.mbar if hasattr(self,'progress') else range(n):\n", " self.epoch = i #To keep track of progress on mbar since the progress callback will use self.epoch\n", " aug_preds.append(self.get_preds(dl=dl, inner=True)[0][None])\n", " aug_preds = torch.cat(aug_preds)\n", " aug_preds = aug_preds.max(0)[0] if use_max else aug_preds.mean(0)\n", " self.epoch = n\n", " with dl.dataset.set_split_idx(1): preds,targs = self.get_preds(dl=dl, inner=True)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#export\n", "@patch\n", "def tta(self:Learner, ds_idx=1, dl=None, n=4, item_tfms=None, batch_tfms=None, beta=0.25, use_max=False):\n", " \"Return predictions on the `ds_idx` dataset or `dl` using Test Time Augmentation\"\n", " if dl is None: dl = self.dls[ds_idx]\n", " if item_tfms is not None or batch_tfms is not None: dl = dl.new(after_item=item_tfms, after_batch=batch_tfms)\n", " try:\n", " self(_before_epoch)\n", " with dl.dataset.set_split_idx(0), self.no_mbar():\n", " if hasattr(self,'progress'): self.progress.mbar = master_bar(list(range(n)))\n", " aug_preds = []\n", " for i in self.progress.mbar if hasattr(self,'progress') else range(n):\n", " self.epoch = i #To keep track of progress on mbar since the progress callback will use self.epoch\n", " aug_preds.append(self.get_preds(dl=dl, inner=True)[0][None])\n", " aug_preds = torch.cat(aug_preds)\n", " aug_preds = aug_preds.max(0)[0] if use_max else aug_preds.mean(0)\n", " self.epoch = n\n", " with dl.dataset.set_split_idx(1): preds,targs = self.get_preds(dl=dl, inner=True)\n", " except CancelFitException: self(event.after_cancel_fit)\n", " finally: self(event.after_fit)\n", "\n", " if use_max: return torch.stack([preds, aug_preds], 0).max(0)[0],targs\n", " preds = (aug_preds,preds) if beta is None else torch.lerp(aug_preds, preds, beta)\n", " return preds,targs" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "In practice, we get the predictions `n` times with the transforms of the training set and average those. The final predictions are `(1-beta)` multiplied by this average + `beta` multiplied by the predictions obtained with the transforms of the dataset. Set `beta` to `None` to get a tuple of the predictions and tta results. You can also use the maximum of all predictions instead of an average by setting `use_max=True`.\n", "\n", "If you want to use new transforms, you can pass them with `item_tfms` and `batch_tfms`." ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#hide\n", "learn = synth_learner()\n", "dl = TfmdDL(Datasets(torch.arange(50)))\n", "learn.dls = DataLoaders(dl, dl)\n", "preds,targs = learn.tta()" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "## Gather arguments" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#export\n", "@patch\n", "def gather_args(self:Learner):\n", " \"Gather config parameters accessible to the learner\"\n", " # init_args\n", " cb_args = {k:v for cb in self.cbs for k,v in getattr(cb,'init_args',{}).items()}\n", " args = {**getattr(self,'init_args',{}), **cb_args, **getattr(self.dls,'init_args',{}),\n", " **getattr(self.opt,'init_args',{}), **getattr(self.loss_func,'init_args',{})}\n", " # callbacks used\n", " args.update({f'{cb}':True for cb in self.cbs})\n", " # input dimensions\n", " try:\n", " n_inp = self.dls.train.n_inp\n", " args['n_inp'] = n_inp\n", " xb = self.dls.train.one_batch()[:n_inp]\n", " args.update({f'input {n+1} dim {i+1}':d for n in range(n_inp) for i,d in enumerate(list(detuplify(xb[n]).shape))})\n", " except: print(f'Could not gather input dimensions')\n", " # other useful information\n", " with ignore_exceptions(): args['batch size'] = self.dls.bs\n", " with ignore_exceptions(): args['batch per epoch'] = len(self.dls.train)\n", " with ignore_exceptions(): args['model parameters'] = total_params(self.model)[0]\n", " with ignore_exceptions(): args['loss function'] = f'{self.loss_func}'\n", " with ignore_exceptions(): args['device'] = self.dls.device.type\n", " with ignore_exceptions(): args['optimizer'] = self.opt_func.__name__\n", " with ignore_exceptions(): args['frozen'] = bool(self.opt.frozen_idx)\n", " with ignore_exceptions(): args['frozen idx'] = self.opt.frozen_idx\n", " with ignore_exceptions(): args['dataset.tfms'] = f'{self.dls.dataset.tfms}'\n", " with ignore_exceptions(): args['dls.after_item'] = f'{self.dls.after_item}'\n", " with ignore_exceptions(): args['dls.before_batch'] = f'{self.dls.before_batch}'\n", " with ignore_exceptions(): args['dls.after_batch'] = f'{self.dls.after_batch}'\n", " return args" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "learn = synth_learner(lr=1e-2)\n", "test_eq(learn.init_args['Learner.__init__.lr'], 0.01)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "## Export -" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "name": "stdout", "output_type": "stream", "text": [ "Converted 00_torch_core.ipynb.\n", "Converted 01_layers.ipynb.\n", "Converted 02_data.load.ipynb.\n", "Converted 03_data.core.ipynb.\n", "Converted 04_data.external.ipynb.\n", "Converted 05_data.transforms.ipynb.\n", "Converted 06_data.block.ipynb.\n", "Converted 07_vision.core.ipynb.\n", "Converted 08_vision.data.ipynb.\n", "Converted 09_vision.augment.ipynb.\n", "Converted 09b_vision.utils.ipynb.\n", "Converted 09c_vision.widgets.ipynb.\n", "Converted 10_tutorial.pets.ipynb.\n", "Converted 11_vision.models.xresnet.ipynb.\n", "Converted 12_optimizer.ipynb.\n", "Converted 13_callback.core.ipynb.\n", "Converted 13a_learner.ipynb.\n", "Converted 13b_metrics.ipynb.\n", "Converted 14_callback.schedule.ipynb.\n", "Converted 14a_callback.data.ipynb.\n", "Converted 15_callback.hook.ipynb.\n", "Converted 15a_vision.models.unet.ipynb.\n", "Converted 16_callback.progress.ipynb.\n", "Converted 17_callback.tracker.ipynb.\n", "Converted 18_callback.fp16.ipynb.\n", "Converted 18a_callback.training.ipynb.\n", "Converted 19_callback.mixup.ipynb.\n", "Converted 20_interpret.ipynb.\n", "Converted 20a_distributed.ipynb.\n", "Converted 21_vision.learner.ipynb.\n", "Converted 22_tutorial.imagenette.ipynb.\n", "Converted 23_tutorial.vision.ipynb.\n", "Converted 24_tutorial.siamese.ipynb.\n", "Converted 24_vision.gan.ipynb.\n", "Converted 30_text.core.ipynb.\n", "Converted 31_text.data.ipynb.\n", "Converted 32_text.models.awdlstm.ipynb.\n", "Converted 33_text.models.core.ipynb.\n", "Converted 34_callback.rnn.ipynb.\n", "Converted 35_tutorial.wikitext.ipynb.\n", "Converted 36_text.models.qrnn.ipynb.\n", "Converted 37_text.learner.ipynb.\n", "Converted 38_tutorial.text.ipynb.\n", "Converted 39_tutorial.transformers.ipynb.\n", "Converted 40_tabular.core.ipynb.\n", "Converted 41_tabular.data.ipynb.\n", "Converted 42_tabular.model.ipynb.\n", "Converted 43_tabular.learner.ipynb.\n", "Converted 44_tutorial.tabular.ipynb.\n", "Converted 45_collab.ipynb.\n", "Converted 46_tutorial.collab.ipynb.\n", "Converted 50_tutorial.datablock.ipynb.\n", "Converted 60_medical.imaging.ipynb.\n", "Converted 61_tutorial.medical_imaging.ipynb.\n", "Converted 65_medical.text.ipynb.\n", "Converted 70_callback.wandb.ipynb.\n", "Converted 71_callback.tensorboard.ipynb.\n", "Converted 72_callback.neptune.ipynb.\n", "Converted 73_callback.captum.ipynb.\n", "Converted 74_callback.cutmix.ipynb.\n", "Converted 97_test_utils.ipynb.\n", "Converted 99_pytorch_doc.ipynb.\n", "Converted index.ipynb.\n", "Converted tutorial.ipynb.\n" ] } ], "source": [ "#hide\n", "from nbdev.export import notebook2script\n", "notebook2script()" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [] } ], "metadata": { "jupytext": { "split_at_heading": true }, "kernelspec": { "display_name": "Python 3", "language": "python", "name": "python3" } }, "nbformat": 4, "nbformat_minor": 4 }