Skip to content
New issue

Have a question about this project? Sign up for a free GitHub account to open an issue and contact its maintainers and the community.

By clicking “Sign up for GitHub”, you agree to our terms of service and privacy statement. We’ll occasionally send you account related emails.

Already on GitHub? Sign in to your account

New Feature: add DropBlock layer #5416

Merged
merged 45 commits into from
Feb 27, 2022
Merged
Show file tree
Hide file tree
Changes from 36 commits
Commits
Show all changes
45 commits
Select commit Hold shift + click to select a range
c33246e
Create dropblock.py
xiaohu2015 Feb 13, 2022
ff34c8e
add dropblock2d
xiaohu2015 Feb 19, 2022
2a86d77
fix pylint
xiaohu2015 Feb 19, 2022
a90e036
refactor dropblock
xiaohu2015 Feb 20, 2022
09f1396
add dropblock
xiaohu2015 Feb 20, 2022
f279981
Rename dropblock.py to drop_block.py
xiaohu2015 Feb 20, 2022
f8cb184
Merge branch 'pytorch:main' into main
xiaohu2015 Feb 20, 2022
ade32f0
fix pylint
xiaohu2015 Feb 20, 2022
2f7a10d
add dropblock
xiaohu2015 Feb 20, 2022
29edef5
add dropblock3d
xiaohu2015 Feb 20, 2022
9969e96
add drop_block3d
xiaohu2015 Feb 20, 2022
bb5be85
Merge branch 'main' into main
xiaohu2015 Feb 20, 2022
a6900f6
Merge branch 'pytorch:main' into main
xiaohu2015 Feb 21, 2022
e5c505e
add dropblock
xiaohu2015 Feb 21, 2022
5ba51be
Update drop_block.py
xiaohu2015 Feb 21, 2022
2901eff
Update torchvision/ops/drop_block.py
xiaohu2015 Feb 21, 2022
90f86f6
Update torchvision/ops/drop_block.py
xiaohu2015 Feb 21, 2022
918c979
Update torchvision/ops/drop_block.py
xiaohu2015 Feb 21, 2022
77ea0ab
Update torchvision/ops/drop_block.py
xiaohu2015 Feb 21, 2022
fefa74e
Update drop_block.py
xiaohu2015 Feb 21, 2022
f5b79ee
Update drop_block.py
xiaohu2015 Feb 21, 2022
8c84c73
Merge branch 'pytorch:main' into main
xiaohu2015 Feb 21, 2022
b45a9e6
import torch.fx
xiaohu2015 Feb 21, 2022
c669853
fix lint
xiaohu2015 Feb 21, 2022
892f1e5
fix lint
xiaohu2015 Feb 21, 2022
7c5e909
Update drop_block.py
xiaohu2015 Feb 21, 2022
d06bc24
improve dropblock
xiaohu2015 Feb 21, 2022
fdac2f4
add dropblock
xiaohu2015 Feb 21, 2022
aedd5f0
refactor dropblock
xiaohu2015 Feb 21, 2022
af7305e
fix doc
xiaohu2015 Feb 21, 2022
2dd89af
remove the limitation of block_size
xiaohu2015 Feb 22, 2022
4f40274
Update torchvision/ops/drop_block.py
xiaohu2015 Feb 22, 2022
b1f91e5
fix lint
xiaohu2015 Feb 22, 2022
60cf559
fix lint
xiaohu2015 Feb 22, 2022
2b3d9cc
add dropblock
xiaohu2015 Feb 22, 2022
4019e7a
Fix linter
datumbox Feb 22, 2022
df0001a
Merge branch 'main' into main
xiaohu2015 Feb 23, 2022
dcf9296
add dropblock random check
xiaohu2015 Feb 23, 2022
84cd3dc
reduce test time
xiaohu2015 Feb 23, 2022
b159f4d
Update test_ops.py
xiaohu2015 Feb 23, 2022
ebea539
speed the dropblock test
xiaohu2015 Feb 23, 2022
6ba5147
Merge branch 'main' into main
datumbox Feb 23, 2022
8d89128
fix lint
xiaohu2015 Feb 23, 2022
bbb9016
Merge branch 'main' into main
datumbox Feb 25, 2022
2af4162
Merge branch 'main' into main
datumbox Feb 27, 2022
File filter

Filter by extension

Filter by extension

Conversations
Failed to load comments.
Loading
Jump to
Jump to file
Failed to load files.
Loading
Diff view
Diff view
4 changes: 4 additions & 0 deletions docs/source/ops.rst
Original file line number Diff line number Diff line change
Expand Up @@ -20,6 +20,8 @@ Operators
box_iou
clip_boxes_to_image
deform_conv2d
drop_block2d
drop_block3d
generalized_box_iou
generalized_box_iou_loss
masks_to_boxes
Expand Down Expand Up @@ -47,3 +49,5 @@ Operators
FrozenBatchNorm2d
ConvNormActivation
SqueezeExcitation
DropBlock2d
DropBlock3d
61 changes: 61 additions & 0 deletions test/test_ops.py
Original file line number Diff line number Diff line change
Expand Up @@ -2,6 +2,7 @@
import os
from abc import ABC, abstractmethod
from functools import lru_cache
from itertools import product
from typing import Callable, List, Tuple

import numpy as np
Expand Down Expand Up @@ -57,6 +58,16 @@ def forward(self, a):
self.layer(a)


class DropBlockWrapper(nn.Module):
def __init__(self, obj):
super().__init__()
self.layer = obj
self.n_inputs = 1

def forward(self, a):
self.layer(a)


class RoIOpTester(ABC):
dtype = torch.float64

Expand Down Expand Up @@ -1357,5 +1368,55 @@ def test_split_normalization_params(self, norm_layer):
assert len(params[1]) == 82


class TestDropBlock:
@pytest.mark.parametrize("seed", range(10))
@pytest.mark.parametrize("dim", (2, 3))
@pytest.mark.parametrize("p", (0, 0.2, 0.5, 0.8))
@pytest.mark.parametrize("block_size", [5, 7, 9, 11])
datumbox marked this conversation as resolved.
Show resolved Hide resolved
@pytest.mark.parametrize("inplace", [True, False])
def test_drop_block(self, seed, dim, p, block_size, inplace):
datumbox marked this conversation as resolved.
Show resolved Hide resolved
torch.manual_seed(seed)
batch_size = 5
channels = 3
height = 11
width = height
depth = height
if dim == 2:
x = torch.ones(size=(batch_size, channels, height, width))
layer = ops.DropBlock2d(p=p, block_size=block_size, inplace=inplace)
feature_size = height * width
elif dim == 3:
x = torch.ones(size=(batch_size, channels, depth, height, width))
layer = ops.DropBlock3d(p=p, block_size=block_size, inplace=inplace)
feature_size = depth * height * width
layer.__repr__()

out = layer(x)
if p == 0:
assert out.equal(x)
if block_size == height:
for b, c in product(range(batch_size), range(channels)):
assert out[b, c].count_nonzero() in (0, feature_size)
datumbox marked this conversation as resolved.
Show resolved Hide resolved

def make_obj(self, dim, p, block_size, inplace, wrap=False):
if dim == 2:
obj = ops.DropBlock2d(p, block_size, inplace)
elif dim == 3:
obj = ops.DropBlock3d(p, block_size, inplace)
return DropBlockWrapper(obj) if wrap else obj

@pytest.mark.parametrize("dim", (2, 3))
@pytest.mark.parametrize("p", (0, 1))
@pytest.mark.parametrize("block_size", [5, 7, 9, 11])
@pytest.mark.parametrize("inplace", [True, False])
def test_is_leaf_node(self, dim, p, block_size, inplace):
op_obj = self.make_obj(dim, p, block_size, inplace, wrap=True)
graph_node_names = get_graph_node_names(op_obj)

assert len(graph_node_names) == 2
assert len(graph_node_names[0]) == len(graph_node_names[1])
assert len(graph_node_names[0]) == 1 + op_obj.n_inputs


if __name__ == "__main__":
pytest.main([__file__])
5 changes: 5 additions & 0 deletions torchvision/ops/__init__.py
Original file line number Diff line number Diff line change
Expand Up @@ -11,6 +11,7 @@
)
from .boxes import box_convert
from .deform_conv import deform_conv2d, DeformConv2d
from .drop_block import drop_block2d, DropBlock2d, drop_block3d, DropBlock3d
from .feature_pyramid_network import FeaturePyramidNetwork
from .focal_loss import sigmoid_focal_loss
from .giou_loss import generalized_box_iou_loss
Expand Down Expand Up @@ -54,4 +55,8 @@
"ConvNormActivation",
"SqueezeExcitation",
"generalized_box_iou_loss",
"drop_block2d",
"DropBlock2d",
"drop_block3d",
"DropBlock3d",
]
155 changes: 155 additions & 0 deletions torchvision/ops/drop_block.py
Original file line number Diff line number Diff line change
@@ -0,0 +1,155 @@
import torch
import torch.fx
import torch.nn.functional as F
from torch import nn, Tensor

from ..utils import _log_api_usage_once


def drop_block2d(
input: Tensor, p: float, block_size: int, inplace: bool = False, eps: float = 1e-06, training: bool = True
) -> Tensor:
"""
Implements DropBlock2d from `"DropBlock: A regularization method for convolutional networks"
<https://arxiv.org/abs/1810.12890>`.

Args:
input (Tensor[N, C, H, W]): The input tensor or 4-dimensions with the first one
being its batch i.e. a batch with ``N`` rows.
p (float): Probability of an element to be dropped.
block_size (int): Size of the block to drop.
inplace (bool): If set to ``True``, will do this operation in-place. Default: ``False``.
eps (float): A value added to the denominator for numerical stability. Default: 1e-6.
training (bool): apply dropblock if is ``True``. Default: ``True``.

Returns:
Tensor[N, C, H, W]: The randomly zeroed tensor after dropblock.
"""
if not torch.jit.is_scripting() and not torch.jit.is_tracing():
_log_api_usage_once(drop_block2d)
if p < 0.0 or p > 1.0:
raise ValueError(f"drop probability has to be between 0 and 1, but got {p}.")
if input.ndim != 4:
raise ValueError(f"input should be 4 dimensional. Got {input.ndim} dimensions.")
if not training or p == 0.0:
return input

N, C, H, W = input.size()
block_size = min(block_size, W, H)
# compute the gamma of Bernoulli distribution
gamma = (p * H * W) / ((block_size ** 2) * ((H - block_size + 1) * (W - block_size + 1)))
noise = torch.empty((N, C, H - block_size + 1, W - block_size + 1), dtype=input.dtype, device=input.device)
noise.bernoulli_(gamma)

noise = F.pad(noise, [block_size // 2] * 4, value=0)
noise = F.max_pool2d(noise, stride=(1, 1), kernel_size=(block_size, block_size), padding=block_size // 2)
noise = 1 - noise
normalize_scale = noise.numel() / (eps + noise.sum())
if inplace:
input.mul_(noise).mul_(normalize_scale)
else:
input = input * noise * normalize_scale
return input


def drop_block3d(
input: Tensor, p: float, block_size: int, inplace: bool = False, eps: float = 1e-06, training: bool = True
) -> Tensor:
"""
Implements DropBlock3d from `"DropBlock: A regularization method for convolutional networks"
<https://arxiv.org/abs/1810.12890>`.

Args:
input (Tensor[N, C, D, H, W]): The input tensor or 5-dimensions with the first one
being its batch i.e. a batch with ``N`` rows.
p (float): Probability of an element to be dropped.
block_size (int): Size of the block to drop.
inplace (bool): If set to ``True``, will do this operation in-place. Default: ``False``.
eps (float): A value added to the denominator for numerical stability. Default: 1e-6.
training (bool): apply dropblock if is ``True``. Default: ``True``.

Returns:
Tensor[N, C, D, H, W]: The randomly zeroed tensor after dropblock.
"""
if not torch.jit.is_scripting() and not torch.jit.is_tracing():
_log_api_usage_once(drop_block3d)
if p < 0.0 or p > 1.0:
raise ValueError(f"drop probability has to be between 0 and 1, but got {p}.")
if input.ndim != 5:
raise ValueError(f"input should be 5 dimensional. Got {input.ndim} dimensions.")
if not training or p == 0.0:
return input

N, C, D, H, W = input.size()
block_size = min(block_size, D, H, W)
# compute the gamma of Bernoulli distribution
gamma = (p * D * H * W) / ((block_size ** 3) * ((D - block_size + 1) * (H - block_size + 1) * (W - block_size + 1)))
noise = torch.empty(
(N, C, D - block_size + 1, H - block_size + 1, W - block_size + 1), dtype=input.dtype, device=input.device
)
noise.bernoulli_(gamma)

noise = F.pad(noise, [block_size // 2] * 6, value=0)
noise = F.max_pool3d(
noise, stride=(1, 1, 1), kernel_size=(block_size, block_size, block_size), padding=block_size // 2
)
noise = 1 - noise
normalize_scale = noise.numel() / (eps + noise.sum())
if inplace:
input.mul_(noise).mul_(normalize_scale)
else:
input = input * noise * normalize_scale
return input


torch.fx.wrap("drop_block2d")


class DropBlock2d(nn.Module):
"""
See :func:`drop_block2d`.
"""

def __init__(self, p: float, block_size: int, inplace: bool = False, eps: float = 1e-06) -> None:
super().__init__()

self.p = p
self.block_size = block_size
self.inplace = inplace
self.eps = eps

def forward(self, input: Tensor) -> Tensor:
"""
Args:
input (Tensor): Input feature map on which some areas will be randomly
dropped.
Returns:
Tensor: The tensor after DropBlock layer.
"""
return drop_block2d(input, self.p, self.block_size, self.inplace, self.eps, self.training)

def __repr__(self) -> str:
s = f"{self.__class__.__name__}(p={self.p}, block_size={self.block_size}, inplace={self.inplace})"
return s


xiaohu2015 marked this conversation as resolved.
Show resolved Hide resolved
torch.fx.wrap("drop_block3d")


class DropBlock3d(DropBlock2d):
"""
See :func:`drop_block3d`.
"""

def __init__(self, p: float, block_size: int, inplace: bool = False, eps: float = 1e-06) -> None:
super().__init__(p, block_size, inplace, eps)

def forward(self, input: Tensor) -> Tensor:
"""
Args:
input (Tensor): Input feature map on which some areas will be randomly
dropped.
Returns:
Tensor: The tensor after DropBlock layer.
"""
return drop_block3d(input, self.p, self.block_size, self.inplace, self.eps, self.training)