{ "cells": [ { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#default_exp data.core" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#export\n", "from fastai2.torch_basics import *\n", "from fastai2.test import *\n", "from fastai2.data.load import *" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "from nbdev.showdoc import *" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "# Data core\n", "\n", "> Core functionality for gathering data" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "The classes here provide functionality for applying a list of transforms to a set of items (`TfmdList`, `DataSource`) or a `DataLoader` (`TfmdDl`) as well as the base class used to gather the data for model training: `DataBunch`." ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "## TfmdDL -" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#export\n", "@typedispatch\n", "def show_batch(x, y, samples, ctxs=None, max_n=9, **kwargs):\n", " if ctxs is None: ctxs = Inf.nones\n", " for i in range_of(samples[0]):\n", " ctxs = [b.show(ctx=c, **kwargs) for b,c,_ in zip(samples.itemgot(i),ctxs,range(max_n))]\n", " return ctxs" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "`show_batch` is a type-dispatched function that is responsible for showing decoded `samples`. `x` and `y` are the input and the target in the batch to be shown, and are passed along to dispatch on their types. There is a different implementation of `show_batch` if `x` is a `TensorImage` or a `TensorText` for instance (see vision.core or text.data for more details). `ctxs` can be passed but the function is responsible to create them if necessary. `kwargs` depend on the specific implementation." ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#export\n", "@typedispatch\n", "def show_results(x, y, samples, outs, ctxs=None, max_n=9, **kwargs):\n", " if ctxs is None: ctxs = Inf.nones\n", " for i in range(len(samples[0])):\n", " ctxs = [b.show(ctx=c, **kwargs) for b,c,_ in zip(samples.itemgot(i),ctxs,range(max_n))]\n", " for i in range(len(outs[0])):\n", " ctxs = [b.show(ctx=c, **kwargs) for b,c,_ in zip(outs.itemgot(i),ctxs,range(max_n))]\n", " return ctxs" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "`show_results` is a type-dispatched function that is responsible for showing decoded `samples` and their corresponding `outs`. Like in `show_batch`, `x` and `y` are the input and the target in the batch to be shown, and are passed along to dispatch on their types. `ctxs` can be passed but the function is responsible to create them if necessary. `kwargs` depend on the specific implementation." ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#export\n", "_all_ = [\"show_batch\", \"show_results\"]" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#export\n", "_batch_tfms = ('after_item','before_batch','after_batch')" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#export\n", "@delegates()\n", "class TfmdDL(DataLoader):\n", " \"Transformed `DataLoader`\"\n", " def __init__(self, dataset, bs=16, shuffle=False, num_workers=None, **kwargs):\n", " if num_workers is None: num_workers = min(16, defaults.cpus)\n", " for nm in _batch_tfms:\n", " kwargs[nm] = Pipeline(kwargs.get(nm,None), as_item=(nm=='before_batch'))\n", " super().__init__(dataset, bs=bs, shuffle=shuffle, num_workers=num_workers, **kwargs)\n", " for nm in _batch_tfms: kwargs[nm].setup(self)\n", "\n", " def _one_pass(self):\n", " its = self.after_batch(self.do_batch([self.do_item(0)]))\n", " self._device = find_device(its)\n", " self._n_inp = 1 if not isinstance(its, (list,tuple)) or len(its)==1 else len(its)-1\n", " self._retain_dl = partial(retain_types, typs=mapped(type,its))\n", "\n", " def _retain_dl(self,b):\n", " self._one_pass()\n", " # we just replaced ourselves, so this is *not* recursive! :)\n", " return self._retain_dl(b)\n", "\n", " def before_iter(self):\n", " super().before_iter()\n", " split_idx = getattr(self.dataset, 'split_idx', None)\n", " for nm in _batch_tfms:\n", " f = getattr(self,nm)\n", " if isinstance(f,Pipeline): f.split_idx=split_idx\n", "\n", " def decode(self, b): return self.before_batch.decode(self.after_batch.decode(self._retain_dl(b)))\n", " def decode_batch(self, b, max_n=9, full=True): return self._decode_batch(self.decode(b), max_n, full)\n", "\n", " def _decode_batch(self, b, max_n=9, full=True):\n", " f = self.after_item.decode\n", " f = compose(f, partial(getattr(self.dataset,'decode',noop), full = full))\n", " return L(batch_to_samples(b, max_n=max_n)).map(f)\n", "\n", " def _pre_show_batch(self, b, max_n=9):\n", " \"Decode `b` to be ready for `show_batch`\"\n", " b = self.decode(b)\n", " if hasattr(b, 'show'): return b,None,None\n", " its = self._decode_batch(b, max_n, full=False)\n", " if not is_listy(b): b,its = [b],L((o,) for o in its)\n", " return detuplify(b[:self.n_inp]),detuplify(b[self.n_inp:]),its\n", "\n", " def show_batch(self, b=None, max_n=9, ctxs=None, show=True, **kwargs):\n", " if b is None: b = self.one_batch()\n", " if not show: return self._pre_show_batch(b, max_n=max_n)\n", " show_batch(*self._pre_show_batch(b, max_n=max_n), ctxs=ctxs, max_n=max_n, **kwargs)\n", "\n", " def show_results(self, b, out, max_n=9, ctxs=None, show=True, **kwargs):\n", " x,y,its = self.show_batch(b, max_n=max_n, show=False)\n", " b_out = b[:self.n_inp] + (tuple(out) if is_listy(out) else (out,))\n", " x1,y1,outs = self.show_batch(b_out, max_n=max_n, show=False)\n", " res = (x,x1,None,None) if its is None else (x, y, its, outs.itemgot(slice(self.n_inp,None)))\n", " if not show: return res\n", " show_results(*res, ctxs=ctxs, max_n=max_n, **kwargs)\n", " \n", " @property\n", " def device(self):\n", " if not hasattr(self, '_device'): _ = self._one_pass()\n", " return self._device\n", "\n", " @property\n", " def n_inp(self):\n", " if hasattr(self.dataset, 'n_inp'): return self.dataset.n_inp\n", " if not hasattr(self, '_n_inp'): self._one_pass()\n", " return self._n_inp" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "A `TfmdDL` is a `DataLoader` that creates `Pipeline` from a list of `Transform`s for the callbacks `after_item`, `before_batch` and `after_batch`. As a result, it can decode or show a processed `batch`." ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "add_docs(TfmdDL,\n", " decode=\"Decode `b` using `tfms`\",\n", " decode_batch=\"Decode `b` entirely\",\n", " show_batch=\"Show `b` (defaults to `one_batch`), a list of lists of pipeline outputs (i.e. output of a `DataLoader`)\",\n", " show_results=\"Show each item of `b` and `out`\",\n", " before_iter=\"override\")" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "class _Category(int, ShowTitle): pass" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#Test retain type\n", "class NegTfm(Transform):\n", " def encodes(self, x): return torch.neg(x)\n", " def decodes(self, x): return torch.neg(x)\n", " \n", "tdl = TfmdDL([(TensorImage([1]),)] * 4, after_batch=NegTfm(), bs=4, num_workers=4)\n", "b = tdl.one_batch()\n", "test_eq(type(b[0]), TensorImage)\n", "b = (tensor([1.,1.,1.,1.]),)\n", "test_eq(type(tdl.decode_batch(b)[0][0]), TensorImage)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "class A(Transform): \n", " def encodes(self, x): return x \n", " def decodes(self, x): return Int(x) \n", "\n", "@Transform\n", "def f(x)->None: return Tuple((x,x))\n", "\n", "start = torch.arange(50)\n", "test_eq_type(f(2), Tuple((2,2)))" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "a = A()\n", "tdl = TfmdDL(start, after_item=lambda x: (a(x), f(x)), bs=4)\n", "x,y = tdl.one_batch()\n", "test_eq(type(y), Tuple)\n", "\n", "s = tdl.decode_batch((x,y))\n", "test_eq(type(s[0][1]), Tuple)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "tdl = TfmdDL(torch.arange(0,50), after_item=A(), after_batch=NegTfm(), bs=4)\n", "test_eq(tdl.dataset[0], start[0])\n", "test_eq(len(tdl), (50-1)//4+1)\n", "test_eq(tdl.bs, 4)\n", "test_stdout(tdl.show_batch, '0\\n1\\n2\\n3')" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "### Methods" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

DataLoader.one_batch[source]

\n", "\n", "> DataLoader.one_batch()\n", "\n" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(TfmdDL.one_batch)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "tfm = NegTfm()\n", "tdl = TfmdDL(start, after_batch=tfm, bs=4)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "b = tdl.one_batch()\n", "test_eq(tensor([0,-1,-2,-3]), b)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

TfmdDL.decode[source]

\n", "\n", "> TfmdDL.decode(**`b`**)\n", "\n", "Decode `b` using `tfms`" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(TfmdDL.decode)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "test_eq(tdl.decode(b), tensor(0,1,2,3))" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

TfmdDL.decode_batch[source]

\n", "\n", "> TfmdDL.decode_batch(**`b`**, **`max_n`**=*`9`*, **`full`**=*`True`*)\n", "\n", "Decode `b` entirely" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(TfmdDL.decode_batch)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "test_eq(tdl.decode_batch(b), [0,1,2,3])" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

TfmdDL.show_batch[source]

\n", "\n", "> TfmdDL.show_batch(**`b`**=*`None`*, **`max_n`**=*`9`*, **`ctxs`**=*`None`*, **`show`**=*`True`*, **\\*\\*`kwargs`**)\n", "\n", "Show `b` (defaults to `one_batch`), a list of lists of pipeline outputs (i.e. output of a [`DataLoader`](/data.load.html#DataLoader))" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(TfmdDL.show_batch)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "## DataBunch -" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "# export\n", "@docs\n", "class DataBunch(GetAttr):\n", " \"Basic wrapper around several `DataLoader`s.\"\n", " _default='train_dl'\n", "\n", " def __init__(self, *dls, path='.'): self.dls,self.path = dls,Path(path)\n", " def __getitem__(self, i): return self.dls[i]\n", " \n", " def new_empty(self):\n", " dls = [dl.new(dl.dataset.new_empty()) for dl in self.dls]\n", " return type(self)(*dls)\n", "\n", " train_dl,valid_dl = add_props(lambda i,x: x[i])\n", " train_ds,valid_ds = add_props(lambda i,x: x[i].dataset)\n", " \n", " @classmethod\n", " @delegates(TfmdDL.__init__)\n", " def from_dblock(cls, dblock, source, path='.', type_tfms=None, item_tfms=None, batch_tfms=None, **kwargs):\n", " return dblock.databunch(source, path=path, type_tfms=type_tfms, item_tfms=item_tfms, batch_tfms=batch_tfms, **kwargs)\n", "\n", " _docs=dict(__getitem__=\"Retrieve `DataLoader` at `i` (`0` is training, `1` is validation)\",\n", " train_dl=\"Training `DataLoader`\",\n", " valid_dl=\"Validation `DataLoader`\",\n", " train_ds=\"Training `Dataset`\",\n", " valid_ds=\"Validation `Dataset`\",\n", " new_empty=\"Create a new empty version of `self` with the same transforms\",\n", " from_dblock=\"Create a databunch from a given `dblock`\")" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "dbch = DataBunch(tdl,tdl)\n", "x = dbch.train_dl.one_batch()\n", "x2 = first(tdl)\n", "test_eq(x,x2)\n", "x2 = dbch.one_batch()\n", "test_eq(x,x2)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "### Methods" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

DataBunch.__getitem__[source]

\n", "\n", "> DataBunch.__getitem__(**`i`**)\n", "\n", "Retrieve [`DataLoader`](/data.load.html#DataLoader) at `i` (`0` is training, `1` is validation)" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(DataBunch.__getitem__)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "x2 = dbch[0].one_batch()\n", "test_eq(x,x2)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

train_dl[source]

\n", "\n", "Training [`DataLoader`](/data.load.html#DataLoader)" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(DataBunch.train_dl, name=\"train_dl\")" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

valid_dl[source]

\n", "\n", "Validation [`DataLoader`](/data.load.html#DataLoader)" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(DataBunch.valid_dl, name=\"valid_dl\")" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

train_ds[source]

\n", "\n", "Training `Dataset`" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(DataBunch.train_ds, name=\"train_ds\")" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

valid_ds[source]

\n", "\n", "Validation `Dataset`" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(DataBunch.valid_ds, name=\"valid_ds\")" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "## TfmdList -" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#export\n", "class FilteredBase:\n", " \"Base class for lists with subsets\"\n", " _dl_type = TfmdDL\n", " def __init__(self, *args, dl_type=None, **kwargs):\n", " if dl_type is not None: self._dl_type = dl_type\n", " self.databunch = delegates(self._dl_type.__init__)(self.databunch)\n", " super().__init__(*args, **kwargs)\n", "\n", " @property\n", " def n_subsets(self): return len(self.splits)\n", " def _new(self, items, **kwargs): return super()._new(items, splits=self.splits, **kwargs)\n", " def subset(self): raise NotImplemented\n", "\n", " def databunch(self, bs=16, val_bs=None, shuffle_train=True, n=None, path='.', dl_type=None, dl_kwargs=None, **kwargs):\n", " if dl_kwargs is None: dl_kwargs = [{}] * self.n_subsets\n", " ns = self.n_subsets-1\n", " bss = ([None]*(ns+1) if bs is None\n", " else [bs] + [3*bs//2]*ns if val_bs is None\n", " else [bs] + [val_bs]*ns)\n", " shuffles = [shuffle_train] + [False]*ns\n", " if dl_type is None: dl_type = self._dl_type\n", " dls = [dl_type(self.subset(i), bs=b, shuffle=s, drop_last=s, n=n if i==0 else None, **kwargs, **dk)\n", " for i,(b,s,dk) in enumerate(zip(bss,shuffles,dl_kwargs))]\n", " return DataBunch(*dls, path=path)\n", "\n", "FilteredBase.train,FilteredBase.valid = add_props(lambda i,x: x.subset(i))" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#export\n", "class TfmdList(FilteredBase, L, GetAttr):\n", " \"A `Pipeline` of `tfms` applied to a collection of `items`\"\n", " _default='tfms'\n", " def __init__(self, items, tfms, use_list=None, do_setup=True, as_item=True, split_idx=None, train_setup=True, splits=None):\n", " super().__init__(items, use_list=use_list)\n", " self.splits = L([slice(None),[]] if splits is None else splits).map(mask2idxs)\n", " if isinstance(tfms,TfmdList): tfms = tfms.tfms\n", " if isinstance(tfms,Pipeline): do_setup=False\n", " self.tfms = Pipeline(tfms, as_item=as_item, split_idx=split_idx)\n", " if do_setup: self.setup(train_setup=train_setup)\n", "\n", " def _new(self, items, **kwargs): return super()._new(items, tfms=self.tfms, do_setup=False, **kwargs)\n", " def subset(self, i): return self._new(self._get(self.splits[i]), split_idx=i)\n", " def _after_item(self, o): return self.tfms(o)\n", " def __repr__(self): return f\"{self.__class__.__name__}: {self.items}\\ntfms - {self.tfms.fs}\"\n", " def __iter__(self): return (self[i] for i in range(len(self)))\n", " def show(self, o, **kwargs): return self.tfms.show(o, **kwargs)\n", " def decode(self, o, **kwargs): return self.tfms.decode(o, **kwargs)\n", " def __call__(self, o, **kwargs): return self.tfms.__call__(o, **kwargs)\n", " def setup(self, train_setup=True): self.tfms.setup(getattr(self,'train',self) if train_setup else self)\n", " def overlapping_splits(self): return L(Counter(self.splits.concat()).values()).filter(gt(1))\n", "\n", " def __getitem__(self, idx):\n", " res = super().__getitem__(idx)\n", " if self._after_item is None: return res\n", " return self._after_item(res) if is_indexer(idx) else res.map(self._after_item)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "add_docs(TfmdList,\n", " setup=\"Transform setup with self\",\n", " decode=\"From `Pipeline\",\n", " show=\"From `Pipeline\",\n", " overlapping_splits=\"All splits that are in more than one split\",\n", " subset=\"New `TfmdList` with same tfms that only includes items in `i`th split\")" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#exports\n", "def decode_at(o, idx):\n", " \"Decoded item at `idx`\"\n", " return o.decode(o[idx])" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#exports\n", "def show_at(o, idx, **kwargs):\n", " \"Show item at `idx`\",\n", " return o.show(o[idx], **kwargs)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "A `TfmdList` combines a collection of object with a `Pipeline`. `tfms` can either be a `Pipeline` or a list of transforms, in which case, it will wrap them in a `Pipeline`. `use_list` is passed along to `L` with the `items`, `as_item` and `split_idx` are passed to each transform of the `Pipeline`. `do_setup` indicates if the `Pipeline.setup` method should be called during initialization." ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "class _IntFloatTfm(Transform):\n", " def encodes(self, o): return Int(o)\n", " def decodes(self, o): return Float(o)\n", "int2f_tfm=_IntFloatTfm()\n", "\n", "def _neg(o): return -o\n", "neg_tfm = Transform(_neg, _neg)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/plain": [ "TfmdList: [1.0, 2.0, 3.0]\n", "tfms - (#2) [Transform: True (object,object) -> _neg (object,object) -> _neg,_IntFloatTfm: True (object,object) -> encodes (object,object) -> decodes]" ] }, "execution_count": null, "metadata": {}, "output_type": "execute_result" } ], "source": [ "items = L([1.,2.,3.]); tfms = [neg_tfm, int2f_tfm]\n", "tl = TfmdList(items, tfms=tfms)\n", "test_eq_type(tl[0], Int(-1))\n", "test_eq_type(tl[1], Int(-2))\n", "test_eq_type(tl.decode(tl[2]), Float(3.))\n", "test_stdout(lambda: show_at(tl, 2), '-3')\n", "tl" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "# add splits to TfmdList\n", "splits = [[0,2],[1]]\n", "tl = TfmdList(items, tfms=tfms, splits=splits)\n", "test_eq(tl.n_subsets, 2)\n", "test_eq(tl.train, tl.subset(0))\n", "test_eq(tl.valid, tl.subset(1))\n", "test_eq(tl.train.items, items[splits[0]])\n", "test_eq(tl.valid.items, items[splits[1]])\n", "test_eq(tl.train.tfms.split_idx, 0)\n", "test_eq(tl.valid.tfms.split_idx, 1)\n", "test_eq_type(tl.splits, L(splits))\n", "assert not tl.overlapping_splits()" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "df = pd.DataFrame(dict(a=[1,2,3],b=[2,3,4]))\n", "tl = TfmdList(df, lambda o: o.a+1, splits=[[0],[1,2]])\n", "test_eq(tl[1,2], [3,4])\n", "tr = tl.subset(0)\n", "test_eq(tr[:], [2])\n", "val = tl.subset(1)\n", "test_eq(val[:], [3,4])" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "class _B(Transform):\n", " def __init__(self): self.m = 0\n", " def encodes(self, o): return o+self.m\n", " def decodes(self, o): return o-self.m\n", " def setups(self, items): self.m = tensor(items).float().mean().item()\n", "\n", "# test for setup, which updates `self.m`\n", "tl = TfmdList(items, _B())\n", "test_eq(tl.m, 2)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "Here's how we can use `TfmdList.setup` to implement a simple category list, getting labels from a mock file list:" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "class _Cat(Transform):\n", " order = 1\n", " def encodes(self, o): return int(self.o2i[o])\n", " def decodes(self, o): return Str(self.vocab[o])\n", " def setups(self, items): self.vocab,self.o2i = uniqueify(L(items), sort=True, bidir=True)\n", "tcat = _Cat()\n", "\n", "def _lbl(o): return Str(o.split('_')[0])\n", "\n", "# Check that tfms are sorted by `order` & `_lbl` is called first\n", "fns = ['dog_0.jpg','cat_0.jpg','cat_2.jpg','cat_1.jpg','dog_1.jpg']\n", "tl = TfmdList(fns, [tcat,_lbl])\n", "exp_voc = ['cat','dog']\n", "test_eq(tcat.vocab, exp_voc)\n", "test_eq(tl.tfms.vocab, exp_voc)\n", "test_eq(tl.vocab, exp_voc)\n", "test_eq(tl, (1,0,0,0,1))\n", "test_eq([tl.decode(o) for o in tl], ('dog','cat','cat','cat','dog'))" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#Check only the training set is taken into account for setup\n", "tl = TfmdList(fns, [tcat,_lbl], splits=[[0,4], [1,2,3]])\n", "test_eq(tcat.vocab, ['dog'])" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "tfm = NegTfm(split_idx=1)\n", "tds = TfmdList(start, A())\n", "tdl = TfmdDL(tds, after_batch=tfm, bs=4)\n", "x = tdl.one_batch()\n", "test_eq(x, torch.arange(4))\n", "tds.split_idx = 1\n", "x = tdl.one_batch()\n", "test_eq(x, -torch.arange(4))\n", "tds.split_idx = 0\n", "x = tdl.one_batch()\n", "test_eq(x, torch.arange(4))" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "tds = TfmdList(start, A())\n", "tdl = TfmdDL(tds, after_batch=NegTfm(), bs=4)\n", "test_eq(tdl.dataset[0], start[0])\n", "test_eq(len(tdl), (len(tds)-1)//4+1)\n", "test_eq(tdl.bs, 4)\n", "test_stdout(tdl.show_batch, '0\\n1\\n2\\n3')" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

TfmdList.subset[source]

\n", "\n", "> TfmdList.subset(**`i`**)\n", "\n", "New [`TfmdList`](/data.core.html#TfmdList) with same tfms that only includes items in `i`th split" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(TfmdList.subset)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "## DataSource -" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#export\n", "@docs\n", "@delegates(TfmdList)\n", "class DataSource(FilteredBase):\n", " \"A dataset that creates a tuple from each `tfms`, passed thru `item_tfms`\"\n", " def __init__(self, items=None, tfms=None, tls=None, n_inp=None, dl_type=None, **kwargs):\n", " super().__init__(dl_type=dl_type)\n", " self.tls = L(tls if tls else [TfmdList(items, t, **kwargs) for t in L(ifnone(tfms,[None]))])\n", " self.n_inp = (1 if len(self.tls)==1 else len(self.tls)-1) if n_inp is None else n_inp\n", "\n", " def __getitem__(self, it):\n", " res = tuple([tl[it] for tl in self.tls])\n", " return res if is_indexer(it) else list(zip(*res))\n", "\n", " def __getattr__(self,k): return gather_attrs(self, k, 'tls')\n", " def __dir__(self): return super().__dir__() + gather_attr_names(self, 'tls')\n", " def __len__(self): return len(self.tls[0])\n", " def __iter__(self): return (self[i] for i in range(len(self)))\n", " def __repr__(self): return coll_repr(self)\n", " def decode(self, o, full=True): return tuple(tl.decode(o_, full=full) for o_,tl in zip(o,tuplify(self.tls, match=o)))\n", " def subset(self, i): return type(self)(tls=L(tl.subset(i) for tl in self.tls), n_inp=self.n_inp)\n", " def _new(self, items, *args, **kwargs): return super()._new(items, tfms=self.tfms, do_setup=False, **kwargs)\n", " def overlapping_splits(self): return self.tls[0].overlapping_splits()\n", " @property\n", " def splits(self): return self.tls[0].splits\n", " @property\n", " def split_idx(self): return self.tls[0].tfms.split_idx\n", " @property\n", " def items(self): return self.tls[0].items\n", " @items.setter\n", " def items(self, v): \n", " for tl in self.tls: tl.items = v\n", "\n", " def show(self, o, ctx=None, **kwargs):\n", " for o_,tl in zip(o,self.tls): ctx = tl.show(o_, ctx=ctx, **kwargs)\n", " return ctx\n", " \n", " def new_empty(self):\n", " tls = [tl._new([self.items[0]], split_idx=tl.split_idx) for tl in self.tls]\n", " return type(self)(tls=tls, n_inp=self.n_inp)\n", " \n", " @contextmanager\n", " def set_split_idx(self, i):\n", " old_split_idx = self.split_idx\n", " for tl in self.tls: tl.tfms.split_idx = i\n", " yield self\n", " for tl in self.tls: tl.tfms.split_idx = old_split_idx\n", "\n", " _docs=dict(\n", " decode=\"Compose `decode` of all `tuple_tfms` then all `tfms` on `i`\",\n", " show=\"Show item `o` in `ctx`\",\n", " databunch=\"Get a `DataBunch`\",\n", " overlapping_splits=\"All splits that are in more than one split\",\n", " subset=\"New `DataSource` that only includes subset `i`\",\n", " new_empty=\"Create a new empty version of the `self`, keeping only the transforms\",\n", " set_split_idx=\"Contextmanager to use the same `DataSource` with another `split_idx`\"\n", " )" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "A `DataSource` creates a tuple from `items` (typically input,target) by applying to them each list of `Transform` (or `Pipeline`) in `tfms`. Note that if `tfms` contains only one list of `tfms`, the items given by `DataSource` will be tuples of one element. \n", "\n", "`n_inp` is the number of elements in the tuples that should be considered part of the input and will default to 1 if `tfms` consists of one set of transforms, `len(tfms)-1` otherwise. In most cases, the number of elements in the tuples spit out by `DataSource` will be 2 (for input,target) but it can happen that there is 3 (Siamese networks or tabular data) in which case we need to be able to determine when the inputs end and the targets begin." ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/plain": [ "(1.0, 2)" ] }, "execution_count": null, "metadata": {}, "output_type": "execute_result" } ], "source": [ "items = [1,2,3,4]\n", "dsrc = DataSource(items, [[neg_tfm,int2f_tfm], [add(1)]])\n", "t = dsrc[0]\n", "test_eq(t, (-1,2))\n", "test_eq(dsrc[0,1,2], [(-1,2),(-2,3),(-3,4)])\n", "test_eq(dsrc.n_inp, 1)\n", "dsrc.decode(t)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "class Norm(Transform):\n", " def encodes(self, o): return (o-self.m)/self.s\n", " def decodes(self, o): return (o*self.s)+self.m\n", " def setups(self, items):\n", " its = tensor(items).float()\n", " self.m,self.s = its.mean(),its.std()" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "items = [1,2,3,4]\n", "nrm = Norm()\n", "dsrc = DataSource(items, [[neg_tfm,int2f_tfm], [neg_tfm,nrm]])\n", "\n", "x,y = zip(*dsrc)\n", "test_close(tensor(y).mean(), 0)\n", "test_close(tensor(y).std(), 1)\n", "test_eq(x, (-1,-2,-3,-4,))\n", "test_eq(nrm.m, -2.5)\n", "test_stdout(lambda:show_at(dsrc, 1), '-2')\n", "\n", "test_eq(dsrc.m, nrm.m)\n", "test_eq(dsrc.norm.m, nrm.m)\n", "test_eq(dsrc.train.norm.m, nrm.m)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#hide\n", "#Check filtering is properly applied\n", "class B(Transform):\n", " def encodes(self, x)->None: return int(x+1)\n", " def decodes(self, x): return Int(x-1)\n", "add1 = B(split_idx=1)\n", "\n", "dsrc = DataSource(items, [neg_tfm, [neg_tfm,int2f_tfm,add1]], splits=[[3],[0,1,2]])\n", "test_eq(dsrc[1], [-2,-2])\n", "test_eq(dsrc.valid[1], [-2,-1])\n", "test_eq(dsrc.valid[[1,1]], [[-2,-1], [-2,-1]])\n", "test_eq(dsrc.train[0], [-4,-4])" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "test_fns = ['dog_0.jpg','cat_0.jpg','cat_2.jpg','cat_1.jpg','kid_1.jpg']\n", "tcat = _Cat()\n", "dsrc = DataSource(test_fns, [[tcat,_lbl]], splits=[[0,1,2], [3,4]])\n", "test_eq(tcat.vocab, ['cat','dog'])\n", "test_eq(dsrc.train, [(1,),(0,),(0,)])\n", "test_eq(dsrc.valid[0], (0,))\n", "test_stdout(lambda: show_at(dsrc.train, 0), \"dog\")" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "inp = [0,1,2,3,4]\n", "dsrc = DataSource(inp, tfms=[None])\n", "\n", "test_eq(*dsrc[2], 2) # Retrieve one item (subset 0 is the default)\n", "test_eq(dsrc[1,2], [(1,),(2,)]) # Retrieve two items by index\n", "mask = [True,False,False,True,False]\n", "test_eq(dsrc[mask], [(0,),(3,)]) # Retrieve two items by mask" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "inp = pd.DataFrame(dict(a=[5,1,2,3,4]))\n", "dsrc = DataSource(inp, tfms=attrgetter('a')).subset(0)\n", "test_eq(*dsrc[2], 2) # Retrieve one item (subset 0 is the default)\n", "test_eq(dsrc[1,2], [(1,),(2,)]) # Retrieve two items by index\n", "mask = [True,False,False,True,False]\n", "test_eq(dsrc[mask], [(5,),(3,)]) # Retrieve two items by mask" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#test n_inp\n", "inp = [0,1,2,3,4]\n", "dsrc = DataSource(inp, tfms=[None])\n", "test_eq(dsrc.n_inp, 1)\n", "dsrc = DataSource(inp, tfms=[[None],[None],[None]])\n", "test_eq(dsrc.n_inp, 2)\n", "dsrc = DataSource(inp, tfms=[[None],[None],[None]], n_inp=1)\n", "test_eq(dsrc.n_inp, 1)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/plain": [ "(#5) [(0,),(1,),(2,),(3,),(4,)]" ] }, "execution_count": null, "metadata": {}, "output_type": "execute_result" } ], "source": [ "# splits can be indices\n", "dsrc = DataSource(range(5), tfms=[None], splits=[tensor([0,2]), [1,3,4]])\n", "\n", "test_eq(dsrc.subset(0), [(0,),(2,)])\n", "test_eq(dsrc.train, [(0,),(2,)]) # Subset 0 is aliased to `train`\n", "test_eq(dsrc.subset(1), [(1,),(3,),(4,)])\n", "test_eq(dsrc.valid, [(1,),(3,),(4,)]) # Subset 1 is aliased to `valid`\n", "test_eq(*dsrc.valid[2], 4)\n", "#assert '[(1,),(3,),(4,)]' in str(dsrc) and '[(0,),(2,)]' in str(dsrc)\n", "dsrc" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "# splits can be boolean masks (they don't have to cover all items, but must be disjoint)\n", "splits = [[False,True,True,False,True], [True,False,False,False,False]]\n", "dsrc = DataSource(range(5), tfms=[None], splits=splits)\n", "\n", "test_eq(dsrc.train, [(1,),(2,),(4,)])\n", "test_eq(dsrc.valid, [(0,)])" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "# apply transforms to all items\n", "tfm = [[lambda x: x*2,lambda x: x+1]]\n", "splits = [[1,2],[0,3,4]]\n", "dsrc = DataSource(range(5), tfm, splits=splits)\n", "test_eq(dsrc.train,[(3,),(5,)])\n", "test_eq(dsrc.valid,[(1,),(7,),(9,)])\n", "test_eq(dsrc.train[False,True], [(5,)])" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "# only transform subset 1\n", "class _Tfm(Transform):\n", " split_idx=1\n", " def encodes(self, x): return x*2\n", " def decodes(self, x): return Str(x//2)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/plain": [ "(#5) [(0,),(1,),(2,),(3,),(4,)]" ] }, "execution_count": null, "metadata": {}, "output_type": "execute_result" } ], "source": [ "dsrc = DataSource(range(5), [_Tfm()], splits=[[1,2],[0,3,4]])\n", "test_eq(dsrc.train,[(1,),(2,)])\n", "test_eq(dsrc.valid,[(0,),(6,),(8,)])\n", "test_eq(dsrc.train[False,True], [(2,)])\n", "dsrc" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#A context manager to change the spli_idx and apply the validation transform on the training set\n", "ds = dsrc.train\n", "with ds.set_split_idx(1):\n", " test_eq(ds,[(2,),(4,)])\n", "test_eq(dsrc.train,[(1,),(2,)])" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#hide\n", "#Test DataSource pickles\n", "dsrc1 = pickle.loads(pickle.dumps(dsrc))\n", "test_eq(dsrc.train, dsrc1.train)\n", "test_eq(dsrc.valid, dsrc1.valid)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "dsrc = DataSource(range(5), [_Tfm(),noop], splits=[[1,2],[0,3,4]])\n", "test_eq(dsrc.train,[(1,1),(2,2)])\n", "test_eq(dsrc.valid,[(0,0),(6,3),(8,4)])" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "start = torch.arange(0,50)\n", "tds = DataSource(start, [A()])\n", "tdl = TfmdDL(tds, after_item=NegTfm(), bs=4)\n", "b = tdl.one_batch()\n", "test_eq(tdl.decode_batch(b), ((0,),(1,),(2,),(3,)))\n", "test_stdout(tdl.show_batch, \"0\\n1\\n2\\n3\")" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "# only transform subset 1\n", "class _Tfm(Transform):\n", " split_idx=1\n", " def encodes(self, x): return x*2\n", "\n", "dsrc = DataSource(range(8), [None], splits=[[1,2,5,7],[0,3,4,6]])\n", "dbch = dsrc.databunch(bs=4, after_batch=_Tfm(), shuffle_train=False)\n", "test_eq(dbch.train_dl, [(tensor([1,2,5, 7]),)])\n", "test_eq(dbch.valid_dl, [(tensor([0,6,8,12]),)])\n", "test_eq(dbch.n_inp, 1)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "### Methods" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "items = [1,2,3,4]\n", "dsrc = DataSource(items, [[neg_tfm,int2f_tfm]])" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

DataSource.decode[source]

\n", "\n", "> DataSource.decode(**`o`**, **`full`**=*`True`*)\n", "\n", "Compose `decode` of all `tuple_tfms` then all `tfms` on `i`" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(DataSource.decode)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "test_eq(*dsrc[0], -1)\n", "test_eq(*dsrc.decode((-1,)), 1)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

DataSource.show[source]

\n", "\n", "> DataSource.show(**`o`**, **`ctx`**=*`None`*, **\\*\\*`kwargs`**)\n", "\n", "Show item `o` in `ctx`" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(DataSource.show)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "test_stdout(lambda:dsrc.show(dsrc[1]), '-2')" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "## Add test set for inference" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "# only transform subset 1\n", "class _Tfm1(Transform):\n", " split_idx=0\n", " def encodes(self, x): return x*3\n", "\n", "dsrc = DataSource(range(8), [[_Tfm(),_Tfm1()]], splits=[[1,2,5,7],[0,3,4,6]])\n", "test_eq(dsrc.train, [(3,),(6,),(15,),(21,)])\n", "test_eq(dsrc.valid, [(0,),(6,),(8,),(12,)])" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#export\n", "def test_set(dsrc, test_items, rm_tfms=0):\n", " \"Create a test set from `test_items` using validation transforms of `dsrc`\"\n", " test_tls = [tl._new(test_items, split_idx=1) for tl in dsrc.tls[:dsrc.n_inp]]\n", " rm_tfms = tuplify(rm_tfms, match=test_tls)\n", " for i,j in enumerate(rm_tfms): test_tls[i].tfms.fs = test_tls[i].tfms.fs[j:]\n", " return DataSource(tls=test_tls)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "class _Tfm1(Transform):\n", " split_idx=0\n", " def encodes(self, x): return x*3\n", "\n", "dsrc = DataSource(range(8), [[_Tfm(),_Tfm1()]], splits=[[1,2,5,7],[0,3,4,6]])\n", "test_eq(dsrc.train, [(3,),(6,),(15,),(21,)])\n", "test_eq(dsrc.valid, [(0,),(6,),(8,),(12,)])\n", "\n", "#Tranform of the validation set are applied\n", "tst = test_set(dsrc, [1,2,3])\n", "test_eq(tst, [(2,),(4,),(6,)])" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#hide\n", "#Test with various input lengths\n", "dsrc = DataSource(range(8), [[_Tfm(),_Tfm1()],[_Tfm(),_Tfm1()],[_Tfm(),_Tfm1()]], splits=[[1,2,5,7],[0,3,4,6]])\n", "tst = test_set(dsrc, [1,2,3])\n", "test_eq(tst, [(2,2),(4,4),(6,6)])\n", "\n", "dsrc = DataSource(range(8), [[_Tfm(),_Tfm1()],[_Tfm(),_Tfm1()],[_Tfm(),_Tfm1()]], splits=[[1,2,5,7],[0,3,4,6]], n_inp=1)\n", "tst = test_set(dsrc, [1,2,3])\n", "test_eq(tst, [(2,),(4,),(6,)])" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#hide\n", "#Test with rm_tfms\n", "dsrc = DataSource(range(8), [[_Tfm(),_Tfm()]], splits=[[1,2,5,7],[0,3,4,6]])\n", "tst = test_set(dsrc, [1,2,3])\n", "test_eq(tst, [(4,),(8,),(12,)])\n", "\n", "dsrc = DataSource(range(8), [[_Tfm(),_Tfm()]], splits=[[1,2,5,7],[0,3,4,6]])\n", "tst = test_set(dsrc, [1,2,3], rm_tfms=1)\n", "test_eq(tst, [(2,),(4,),(6,)])\n", "\n", "dsrc = DataSource(range(8), [[_Tfm(),_Tfm()], [_Tfm(),_Tfm()]], splits=[[1,2,5,7],[0,3,4,6]], n_inp=2)\n", "tst = test_set(dsrc, [1,2,3], rm_tfms=(1,0))\n", "test_eq(tst, [(2,4),(4,8),(6,12)])" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "#export\n", "@delegates(TfmdDL.__init__)\n", "def test_dl(dbunch, test_items, rm_type_tfms=0, **kwargs):\n", " \"Create a test dataloader from `test_items` using validation transforms of `dbunch`\"\n", " test_ds = test_set(dbunch.valid_ds, test_items, rm_tfms=rm_type_tfms) if isinstance(dbunch.valid_ds, DataSource) else test_items\n", " return dbunch.valid_dl.new(test_ds, **kwargs)" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "dsrc = DataSource(range(8), [[_Tfm(),_Tfm1()]], splits=[[1,2,5,7],[0,3,4,6]])\n", "dbunch = dsrc.databunch(bs=4)\n", "tst_dl = test_dl(dbunch, [2,3,4,5])\n", "test_eq(list(tst_dl), [(tensor([ 4, 6, 8, 10]),)])\n", "#Test you can change transforms\n", "tst_dl = test_dl(dbunch, [2,3,4,5], after_item=add1)\n", "test_eq(list(tst_dl), [(tensor([ 5, 7, 9, 11]),)])" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "## Export -" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "name": "stdout", "output_type": "stream", "text": [ "Converted 00_test.ipynb.\n", "Converted 01_core_foundation.ipynb.\n", "Converted 01a_core_utils.ipynb.\n", "Converted 01b_core_dispatch.ipynb.\n", "Converted 01c_core_transform.ipynb.\n", "Converted 02_core_script.ipynb.\n", "Converted 03_torchcore.ipynb.\n", "Converted 03a_layers.ipynb.\n", "Converted 04_data_load.ipynb.\n", "Converted 05_data_core.ipynb.\n", "Converted 06_data_transforms.ipynb.\n", "Converted 07_data_block.ipynb.\n", "Converted 08_vision_core.ipynb.\n", "Converted 09_vision_augment.ipynb.\n", "Converted 09a_vision_data.ipynb.\n", "Converted 09b_vision_utils.ipynb.\n", "Converted 10_pets_tutorial.ipynb.\n", "Converted 11_vision_models_xresnet.ipynb.\n", "Converted 12_optimizer.ipynb.\n", "Converted 13_learner.ipynb.\n", "Converted 13a_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 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_transfer_learning.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_ulmfit.ipynb.\n", "Converted 40_tabular_core.ipynb.\n", "Converted 41_tabular_model.ipynb.\n", "Converted 42_tabular_rapids.ipynb.\n", "Converted 50_data_block_examples.ipynb.\n", "Converted 60_medical_imaging.ipynb.\n", "Converted 65_medical_text.ipynb.\n", "Converted 70_callback_wandb.ipynb.\n", "Converted 71_callback_tensorboard.ipynb.\n", "Converted 90_notebook_core.ipynb.\n", "Converted 91_notebook_export.ipynb.\n", "Converted 92_notebook_showdoc.ipynb.\n", "Converted 93_notebook_export2html.ipynb.\n", "Converted 94_notebook_test.ipynb.\n", "Converted 95_index.ipynb.\n", "Converted 96_data_external.ipynb.\n", "Converted 97_utils_test.ipynb.\n", "Converted notebook2jekyll.ipynb.\n", "Converted xse_resnext.ipynb.\n" ] } ], "source": [ "#hide\n", "from nbdev.export import notebook2script\n", "notebook2script()" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [] } ], "metadata": { "kernelspec": { "display_name": "Python 3", "language": "python", "name": "python3" } }, "nbformat": 4, "nbformat_minor": 2 }