.. DO NOT EDIT. .. THIS FILE WAS AUTOMATICALLY GENERATED BY SPHINX-GALLERY. .. TO MAKE CHANGES, EDIT THE SOURCE PYTHON FILE: .. "tutorials/tensorclass_fashion.py" .. LINE NUMBERS ARE GIVEN BELOW. .. only:: html .. note:: :class: sphx-glr-download-link-note :ref:`Go to the end ` to download the full example code. .. rst-class:: sphx-glr-example-title .. _sphx_glr_tutorials_tensorclass_fashion.py: Using tensorclasses for datasets ================================ .. GENERATED FROM PYTHON SOURCE LINES 7-13 In this tutorial we demonstrate how tensorclasses can be used to efficiently and transparently load and manage data inside a training pipeline. The tutorial is based heavily on the `PyTorch Quickstart Tutorial `__, but modified to demonstrate use of tensorclass. See the related tutorial using ``TensorDict``. .. GENERATED FROM PYTHON SOURCE LINES 13-27 .. code-block:: Python import torch import torch.nn as nn from tensordict import MemoryMappedTensor, tensorclass from torch.utils.data import DataLoader from torchvision import datasets from torchvision.transforms import ToTensor device = "cuda" if torch.cuda.is_available() else "cpu" print(f"Using device: {device}") .. rst-class:: sphx-glr-script-out .. code-block:: none Using device: cpu .. GENERATED FROM PYTHON SOURCE LINES 28-32 The ``torchvision.datasets`` module contains a number of convenient pre-prepared datasets. In this tutorial we'll use the relatively simple FashionMNIST dataset. Each image is an item of clothing, the objective is to classify the type of clothing in the image (e.g. "Bag", "Sneaker" etc.). .. GENERATED FROM PYTHON SOURCE LINES 32-46 .. code-block:: Python training_data = datasets.FashionMNIST( root="data", train=True, download=True, transform=ToTensor(), ) test_data = datasets.FashionMNIST( root="data", train=False, download=True, transform=ToTensor(), ) .. rst-class:: sphx-glr-script-out .. code-block:: none 0%| | 0.00/26.4M [00:00`__. .. GENERATED FROM PYTHON SOURCE LINES 120-144 .. code-block:: Python class Net(nn.Module): def __init__(self): super().__init__() self.flatten = nn.Flatten() self.linear_relu_stack = nn.Sequential( nn.Linear(28 * 28, 512), nn.ReLU(), nn.Linear(512, 512), nn.ReLU(), nn.Linear(512, 10), ) def forward(self, x): x = self.flatten(x) logits = self.linear_relu_stack(x) return logits model = Net().to(device) model_tc = Net().to(device) model, model_tc .. rst-class:: sphx-glr-script-out .. code-block:: none (Net( (flatten): Flatten(start_dim=1, end_dim=-1) (linear_relu_stack): Sequential( (0): Linear(in_features=784, out_features=512, bias=True) (1): ReLU() (2): Linear(in_features=512, out_features=512, bias=True) (3): ReLU() (4): Linear(in_features=512, out_features=10, bias=True) ) ), Net( (flatten): Flatten(start_dim=1, end_dim=-1) (linear_relu_stack): Sequential( (0): Linear(in_features=784, out_features=512, bias=True) (1): ReLU() (2): Linear(in_features=512, out_features=512, bias=True) (3): ReLU() (4): Linear(in_features=512, out_features=10, bias=True) ) )) .. GENERATED FROM PYTHON SOURCE LINES 145-151 Optimizing the parameters --------------------------------- We'll optimise the parameters of the model using stochastic gradient descent and cross-entropy loss. .. GENERATED FROM PYTHON SOURCE LINES 151-176 .. code-block:: Python loss_fn = nn.CrossEntropyLoss() optimizer = torch.optim.SGD(model.parameters(), lr=1e-3) optimizer_tc = torch.optim.SGD(model_tc.parameters(), lr=1e-3) def train(dataloader, model, loss_fn, optimizer): size = len(dataloader.dataset) model.train() for batch, (X, y) in enumerate(dataloader): X, y = X.to(device), y.to(device) pred = model(X) loss = loss_fn(pred, y) optimizer.zero_grad() loss.backward() optimizer.step() if batch % 100 == 0: loss, current = loss.item(), batch * len(X) print(f"loss: {loss:>7f} [{current:>5d}/{size:>5d}]") .. GENERATED FROM PYTHON SOURCE LINES 177-181 The training loop for our tensorclass-based DataLoader is very similar, we just adjust how we unpack the data to the more explicit attribute-based retrieval offered by the tensorclass. The ``.contiguous()`` method loads the data stored in the memmap tensor. .. GENERATED FROM PYTHON SOURCE LINES 181-267 .. code-block:: Python def train_tc(dataloader, model, loss_fn, optimizer): size = len(dataloader.dataset) model.train() for batch, data in enumerate(dataloader): X, y = data.images.contiguous(), data.targets.contiguous() pred = model(X) loss = loss_fn(pred, y) optimizer.zero_grad() loss.backward() optimizer.step() if batch % 100 == 0: loss, current = loss.item(), batch * len(X) print(f"loss: {loss:>7f} [{current:>5d}/{size:>5d}]") def test(dataloader, model, loss_fn): size = len(dataloader.dataset) num_batches = len(dataloader) model.eval() test_loss, correct = 0, 0 with torch.no_grad(): for X, y in dataloader: X, y = X.to(device), y.to(device) pred = model(X) test_loss += loss_fn(pred, y).item() correct += (pred.argmax(1) == y).type(torch.float).sum().item() test_loss /= num_batches correct /= size print( f"Test Error: \n Accuracy: {(100 * correct):>0.1f}%, Avg loss: {test_loss:>8f} \n" ) def test_tc(dataloader, model, loss_fn): size = len(dataloader.dataset) num_batches = len(dataloader) model.eval() test_loss, correct = 0, 0 with torch.no_grad(): for batch in dataloader: X, y = batch.images.contiguous(), batch.targets.contiguous() pred = model(X) test_loss += loss_fn(pred, y).item() correct += (pred.argmax(1) == y).type(torch.float).sum().item() test_loss /= num_batches correct /= size print( f"Test Error: \n Accuracy: {(100 * correct):>0.1f}%, Avg loss: {test_loss:>8f} \n" ) for d in train_dataloader_tc: print(d) break import time t0 = time.time() epochs = 5 for t in range(epochs): print(f"Epoch {t + 1}\n-------------------------") train_tc(train_dataloader_tc, model_tc, loss_fn, optimizer_tc) test_tc(test_dataloader_tc, model_tc, loss_fn) print(f"Tensorclass training done! time: {time.time() - t0: 4.4f} s") t0 = time.time() epochs = 5 for t in range(epochs): print(f"Epoch {t + 1}\n-------------------------") train(train_dataloader, model, loss_fn, optimizer) test(test_dataloader, model, loss_fn) print(f"Training done! time: {time.time() - t0: 4.4f} s") .. rst-class:: sphx-glr-script-out .. code-block:: none FashionMNISTData( images=Tensor(shape=torch.Size([64, 28, 28]), device=cpu, dtype=torch.float32, is_shared=False), targets=Tensor(shape=torch.Size([64]), device=cpu, dtype=torch.int64, is_shared=False), batch_size=torch.Size([64]), device=cpu, is_shared=False) Epoch 1 ------------------------- loss: 2.309098 [ 0/60000] loss: 2.296995 [ 6400/60000] loss: 2.282863 [12800/60000] loss: 2.273057 [19200/60000] loss: 2.263108 [25600/60000] loss: 2.236950 [32000/60000] loss: 2.243644 [38400/60000] loss: 2.216352 [44800/60000] loss: 2.206101 [51200/60000] loss: 2.181821 [57600/60000] Test Error: Accuracy: 43.9%, Avg loss: 2.177676 Epoch 2 ------------------------- loss: 2.184656 [ 0/60000] loss: 2.177800 [ 6400/60000] loss: 2.131632 [12800/60000] loss: 2.144295 [19200/60000] loss: 2.104907 [25600/60000] loss: 2.054239 [32000/60000] loss: 2.082007 [38400/60000] loss: 2.014649 [44800/60000] loss: 2.007189 [51200/60000] loss: 1.952836 [57600/60000] Test Error: Accuracy: 55.1%, Avg loss: 1.947203 Epoch 3 ------------------------- loss: 1.972391 [ 0/60000] loss: 1.947021 [ 6400/60000] loss: 1.846253 [12800/60000] loss: 1.880965 [19200/60000] loss: 1.781960 [25600/60000] loss: 1.735687 [32000/60000] loss: 1.760682 [38400/60000] loss: 1.661476 [44800/60000] loss: 1.672414 [51200/60000] loss: 1.578934 [57600/60000] Test Error: Accuracy: 60.5%, Avg loss: 1.590364 Epoch 4 ------------------------- loss: 1.646736 [ 0/60000] loss: 1.611050 [ 6400/60000] loss: 1.469680 [12800/60000] loss: 1.538767 [19200/60000] loss: 1.422644 [25600/60000] loss: 1.411621 [32000/60000] loss: 1.427902 [38400/60000] loss: 1.348701 [44800/60000] loss: 1.377567 [51200/60000] loss: 1.276364 [57600/60000] Test Error: Accuracy: 62.5%, Avg loss: 1.307161 Epoch 5 ------------------------- loss: 1.378672 [ 0/60000] loss: 1.356292 [ 6400/60000] loss: 1.200504 [12800/60000] loss: 1.299592 [19200/60000] loss: 1.181137 [25600/60000] loss: 1.200128 [32000/60000] loss: 1.216121 [38400/60000] loss: 1.153240 [44800/60000] loss: 1.190223 [51200/60000] loss: 1.098847 [57600/60000] Test Error: Accuracy: 63.8%, Avg loss: 1.128341 Tensorclass training done! time: 8.5029 s Epoch 1 ------------------------- loss: 2.301558 [ 0/60000] loss: 2.287850 [ 6400/60000] loss: 2.267768 [12800/60000] loss: 2.266540 [19200/60000] loss: 2.242177 [25600/60000] loss: 2.222106 [32000/60000] loss: 2.227977 [38400/60000] loss: 2.197102 [44800/60000] loss: 2.187286 [51200/60000] loss: 2.158054 [57600/60000] Test Error: Accuracy: 52.7%, Avg loss: 2.147735 Epoch 2 ------------------------- loss: 2.156160 [ 0/60000] loss: 2.146098 [ 6400/60000] loss: 2.077926 [12800/60000] loss: 2.101998 [19200/60000] loss: 2.049972 [25600/60000] loss: 1.988866 [32000/60000] loss: 2.022804 [38400/60000] loss: 1.937112 [44800/60000] loss: 1.935618 [51200/60000] loss: 1.864654 [57600/60000] Test Error: Accuracy: 59.9%, Avg loss: 1.859219 Epoch 3 ------------------------- loss: 1.888533 [ 0/60000] loss: 1.862214 [ 6400/60000] loss: 1.730510 [12800/60000] loss: 1.784361 [19200/60000] loss: 1.682599 [25600/60000] loss: 1.624530 [32000/60000] loss: 1.659126 [38400/60000] loss: 1.556772 [44800/60000] loss: 1.579414 [51200/60000] loss: 1.473874 [57600/60000] Test Error: Accuracy: 62.3%, Avg loss: 1.492313 Epoch 4 ------------------------- loss: 1.554773 [ 0/60000] loss: 1.528786 [ 6400/60000] loss: 1.369470 [12800/60000] loss: 1.455213 [19200/60000] loss: 1.345353 [25600/60000] loss: 1.327606 [32000/60000] loss: 1.354587 [38400/60000] loss: 1.279378 [44800/60000] loss: 1.311781 [51200/60000] loss: 1.212201 [57600/60000] Test Error: Accuracy: 63.5%, Avg loss: 1.238718 Epoch 5 ------------------------- loss: 1.309647 [ 0/60000] loss: 1.301050 [ 6400/60000] loss: 1.128442 [12800/60000] loss: 1.246816 [19200/60000] loss: 1.122059 [25600/60000] loss: 1.136712 [32000/60000] loss: 1.170572 [38400/60000] loss: 1.109064 [44800/60000] loss: 1.145305 [51200/60000] loss: 1.060434 [57600/60000] Test Error: Accuracy: 64.8%, Avg loss: 1.081106 Training done! time: 34.2454 s .. rst-class:: sphx-glr-timing **Total running time of the script:** (1 minutes 0.560 seconds) .. _sphx_glr_download_tutorials_tensorclass_fashion.py: .. only:: html .. container:: sphx-glr-footer sphx-glr-footer-example .. container:: sphx-glr-download sphx-glr-download-jupyter :download:`Download Jupyter notebook: tensorclass_fashion.ipynb ` .. container:: sphx-glr-download sphx-glr-download-python :download:`Download Python source code: tensorclass_fashion.py ` .. container:: sphx-glr-download sphx-glr-download-zip :download:`Download zipped: tensorclass_fashion.zip ` .. only:: html .. rst-class:: sphx-glr-signature `Gallery generated by Sphinx-Gallery `_