TiCo
Example implementation of Transformation Invariance and Covariance Contrast (TiCo) for self-supervised visual representation learning. Similar to BYOL, this method is based on maximizing the agreement among embeddings of different distorted versions of the same image, which pushes the encoder to produce transformation invariant representations.
This example can be run from the command line with:
python lightly/examples/pytorch/tico.py
# This example requires the following dependencies to be installed:
# pip install lightly
# Note: The model and training settings do not follow the reference settings
# from the paper. The settings are chosen such that the example can easily be
# run on a small dataset with a single GPU.
import copy
import torch
import torchvision
from torch import nn
from lightly.loss.tico_loss import TiCoLoss
from lightly.models.modules.heads import TiCoProjectionHead
from lightly.models.utils import deactivate_requires_grad, update_momentum
from lightly.transforms import TiCoTransform, TiCoView1Transform, TiCoView2Transform
from lightly.utils.scheduler import cosine_schedule
class TiCo(nn.Module):
def __init__(self, backbone):
super().__init__()
self.backbone = backbone
self.projection_head = TiCoProjectionHead(512, 1024, 256)
self.backbone_momentum = copy.deepcopy(self.backbone)
self.projection_head_momentum = copy.deepcopy(self.projection_head)
deactivate_requires_grad(self.backbone_momentum)
deactivate_requires_grad(self.projection_head_momentum)
def forward(self, x):
y = self.backbone(x).flatten(start_dim=1)
z = self.projection_head(y)
return z
def forward_momentum(self, x):
y = self.backbone_momentum(x).flatten(start_dim=1)
z = self.projection_head_momentum(y)
z = z.detach()
return z
resnet = torchvision.models.resnet18()
backbone = nn.Sequential(*list(resnet.children())[:-1])
model = TiCo(backbone)
device = "cuda" if torch.cuda.is_available() else "cpu"
model.to(device)
# TiCo uses the same augmentations as BYOL.
# We disable resizing and gaussian blur for cifar10.
transform = TiCoTransform(
view_1_transform=TiCoView1Transform(input_size=32, gaussian_blur=0.0),
view_2_transform=TiCoView2Transform(input_size=32, gaussian_blur=0.0),
)
dataset = torchvision.datasets.CIFAR10(
"datasets/cifar10", download=True, transform=transform
)
# or create a dataset from a folder containing images or videos:
# dataset = LightlyDataset("path/to/folder", transform=transform)
dataloader = torch.utils.data.DataLoader(
dataset,
batch_size=256,
shuffle=True,
drop_last=True,
num_workers=8,
)
criterion = TiCoLoss()
optimizer = torch.optim.SGD(model.parameters(), lr=0.06)
epochs = 10
print("Starting Training")
for epoch in range(epochs):
total_loss = 0
momentum_val = cosine_schedule(epoch, epochs, 0.996, 1)
for batch in dataloader:
x0, x1 = batch[0]
update_momentum(model.backbone, model.backbone_momentum, m=momentum_val)
update_momentum(
model.projection_head, model.projection_head_momentum, m=momentum_val
)
x0 = x0.to(device)
x1 = x1.to(device)
z0 = model(x0)
z1 = model.forward_momentum(x1)
loss = criterion(z0, z1)
total_loss += loss.detach()
loss.backward()
optimizer.step()
optimizer.zero_grad()
avg_loss = total_loss / len(dataloader)
print(f"epoch: {epoch:>02}, loss: {avg_loss:.5f}")
This example can be run from the command line with:
python lightly/examples/pytorch_lightning/tico.py
# This example requires the following dependencies to be installed:
# pip install lightly
import copy
import pytorch_lightning as pl
import torch
import torchvision
from torch import nn
from lightly.loss.tico_loss import TiCoLoss
from lightly.models.modules.heads import TiCoProjectionHead
from lightly.models.utils import deactivate_requires_grad, update_momentum
from lightly.transforms import TiCoTransform, TiCoView1Transform, TiCoView2Transform
from lightly.utils.scheduler import cosine_schedule
class TiCo(pl.LightningModule):
def __init__(self):
super().__init__()
resnet = torchvision.models.resnet18()
self.backbone = nn.Sequential(*list(resnet.children())[:-1])
self.projection_head = TiCoProjectionHead(512, 1024, 256)
self.backbone_momentum = copy.deepcopy(self.backbone)
self.projection_head_momentum = copy.deepcopy(self.projection_head)
deactivate_requires_grad(self.backbone_momentum)
deactivate_requires_grad(self.projection_head_momentum)
self.criterion = TiCoLoss()
def forward(self, x):
y = self.backbone(x).flatten(start_dim=1)
z = self.projection_head(y)
return z
def forward_momentum(self, x):
y = self.backbone_momentum(x).flatten(start_dim=1)
z = self.projection_head_momentum(y)
z = z.detach()
return z
def training_step(self, batch, batch_idx):
(x0, x1) = batch[0]
momentum = cosine_schedule(self.current_epoch, 10, 0.996, 1)
update_momentum(self.backbone, self.backbone_momentum, m=momentum)
update_momentum(self.projection_head, self.projection_head_momentum, m=momentum)
x0 = x0.to(self.device)
x1 = x1.to(self.device)
z0 = self.forward(x0)
z1 = self.forward_momentum(x1)
loss = self.criterion(z0, z1)
return loss
def configure_optimizers(self):
return torch.optim.SGD(self.parameters(), lr=0.06)
model = TiCo()
# We disable resizing and gaussian blur for cifar10.
transform = TiCoTransform(
view_1_transform=TiCoView1Transform(input_size=32, gaussian_blur=0.0),
view_2_transform=TiCoView2Transform(input_size=32, gaussian_blur=0.0),
)
dataset = torchvision.datasets.CIFAR10(
"datasets/cifar10", download=True, transform=transform
)
# or create a dataset from a folder containing images or videos:
# dataset = LightlyDataset("path/to/folder", transform=transform)
dataloader = torch.utils.data.DataLoader(
dataset,
batch_size=256,
shuffle=True,
drop_last=True,
num_workers=8,
)
accelerator = "gpu" if torch.cuda.is_available() else "cpu"
trainer = pl.Trainer(max_epochs=10, devices=1, accelerator=accelerator)
trainer.fit(model=model, train_dataloaders=dataloader)
This example runs on multiple gpus using Distributed Data Parallel (DDP) training with Pytorch Lightning. At least one GPU must be available on the system. The example can be run from the command line with:
python lightly/examples/pytorch_lightning_distributed/tico.py
The model differs in the following ways from the non-distributed implementation:
Distributed Data Parallel is enabled
Synchronized Batch Norm is used in place of standard Batch Norm
Note that Synchronized Batch Norm is optional and the model can also be trained without it. Without Synchronized Batch Norm the batch norm for each GPU is only calculated based on the features on that specific GPU.
# This example requires the following dependencies to be installed:
# pip install lightly
import copy
import pytorch_lightning as pl
import torch
import torchvision
from torch import nn
from lightly.loss.tico_loss import TiCoLoss
from lightly.models.modules.heads import TiCoProjectionHead
from lightly.models.utils import deactivate_requires_grad, update_momentum
from lightly.transforms import TiCoTransform, TiCoView1Transform, TiCoView2Transform
from lightly.utils.scheduler import cosine_schedule
class TiCo(pl.LightningModule):
def __init__(self):
super().__init__()
resnet = torchvision.models.resnet18()
self.backbone = nn.Sequential(*list(resnet.children())[:-1])
self.projection_head = TiCoProjectionHead(512, 1024, 256)
self.backbone_momentum = copy.deepcopy(self.backbone)
self.projection_head_momentum = copy.deepcopy(self.projection_head)
deactivate_requires_grad(self.backbone_momentum)
deactivate_requires_grad(self.projection_head_momentum)
self.criterion = TiCoLoss()
def forward(self, x):
y = self.backbone(x).flatten(start_dim=1)
z = self.projection_head(y)
return z
def forward_momentum(self, x):
y = self.backbone_momentum(x).flatten(start_dim=1)
z = self.projection_head_momentum(y)
z = z.detach()
return z
def training_step(self, batch, batch_idx):
(x0, x1) = batch[0]
momentum = cosine_schedule(self.current_epoch, 10, 0.996, 1)
update_momentum(self.backbone, self.backbone_momentum, m=momentum)
update_momentum(self.projection_head, self.projection_head_momentum, m=momentum)
x0 = x0.to(self.device)
x1 = x1.to(self.device)
z0 = self.forward(x0)
z1 = self.forward_momentum(x1)
loss = self.criterion(z0, z1)
return loss
def configure_optimizers(self):
return torch.optim.SGD(self.parameters(), lr=0.06)
model = TiCo()
# We disable resizing and gaussian blur for cifar10.
transform = TiCoTransform(
view_1_transform=TiCoView1Transform(input_size=32, gaussian_blur=0.0),
view_2_transform=TiCoView2Transform(input_size=32, gaussian_blur=0.0),
)
dataset = torchvision.datasets.CIFAR10(
"datasets/cifar10", download=True, transform=transform
)
# or create a dataset from a folder containing images or videos:
# dataset = LightlyDataset("path/to/folder", transform=transform)
dataloader = torch.utils.data.DataLoader(
dataset,
batch_size=256,
shuffle=True,
drop_last=True,
num_workers=8,
)
# Train with DDP and use Synchronized Batch Norm for a more accurate batch norm
# calculation. Distributed sampling is also enabled with replace_sampler_ddp=True.
trainer = pl.Trainer(
max_epochs=10,
devices="auto",
accelerator="gpu",
strategy="ddp",
sync_batchnorm=True,
use_distributed_sampler=True, # or replace_sampler_ddp=True for PyTorch Lightning <2.0
)
trainer.fit(model=model, train_dataloaders=dataloader)