Skip to content

Commit 6b8115a

Browse files
author
kuangliu
committed
Update ResNeXt model and accuracy
1 parent 91a5b80 commit 6b8115a

File tree

4 files changed

+69
-68
lines changed

4 files changed

+69
-68
lines changed

README.md

Lines changed: 9 additions & 8 deletions
Original file line numberDiff line numberDiff line change
@@ -13,14 +13,15 @@ Cons:
1313
- No built-in log.
1414

1515
## Accuracy
16-
| Model | Acc. |
17-
| ------------ |----------- |
18-
| VGG16 | 92.64% |
19-
| ResNet18 | 93.02% |
20-
| ResNet50 | 93.62% |
21-
| ResNet101 | 93.75% |
22-
| ResNeXt29 | 94.73% |
23-
| DenseNet121 | 95.04% |
16+
| Model | Acc. |
17+
| ------------ | ----------- |
18+
| VGG16 | 92.64% |
19+
| ResNet18 | 93.02% |
20+
| ResNet50 | 93.62% |
21+
| ResNet101 | 93.75% |
22+
| ResNeXt29(32x4d) | 94.73% |
23+
| ResNeXt29(2x64d) | 94.82% |
24+
| DenseNet121 | 95.04% |
2425

2526
## Learning rate adjustment
2627
I manually change the `lr` during training:

main.py

Lines changed: 1 addition & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -59,7 +59,7 @@
5959
# net = ResNet18()
6060
# net = GoogLeNet()
6161
# net = DenseNet121()
62-
net = ResNeXt29()
62+
net = ResNeXt29_2x64d()
6363

6464
if use_cuda:
6565
net.cuda()

models/resnext.py

Lines changed: 44 additions & 46 deletions
Original file line numberDiff line numberDiff line change
@@ -4,7 +4,6 @@
44
'''
55
import torch
66
import torch.nn as nn
7-
import torch.nn.init as init
87
import torch.nn.functional as F
98

109
from torch.autograd import Variable
@@ -14,20 +13,21 @@ class Block(nn.Module):
1413
'''Grouped convolution block.'''
1514
expansion = 2
1615

17-
def __init__(self, in_planes, planes, stride=1, cardinality=32):
16+
def __init__(self, in_planes, cardinality=32, bottleneck_width=4, stride=1):
1817
super(Block, self).__init__()
19-
self.conv1 = nn.Conv2d(in_planes, planes, kernel_size=1, bias=False)
20-
self.bn1 = nn.BatchNorm2d(planes)
21-
self.conv2 = nn.Conv2d(planes, planes, kernel_size=3, stride=stride, padding=1, groups=cardinality, bias=False)
22-
self.bn2 = nn.BatchNorm2d(planes)
23-
self.conv3 = nn.Conv2d(planes, self.expansion*planes, kernel_size=1, bias=False)
24-
self.bn3 = nn.BatchNorm2d(self.expansion*planes)
18+
group_width = cardinality * bottleneck_width
19+
self.conv1 = nn.Conv2d(in_planes, group_width, kernel_size=1, bias=False)
20+
self.bn1 = nn.BatchNorm2d(group_width)
21+
self.conv2 = nn.Conv2d(group_width, group_width, kernel_size=3, stride=stride, padding=1, groups=cardinality, bias=False)
22+
self.bn2 = nn.BatchNorm2d(group_width)
23+
self.conv3 = nn.Conv2d(group_width, self.expansion*group_width, kernel_size=1, bias=False)
24+
self.bn3 = nn.BatchNorm2d(self.expansion*group_width)
2525

2626
self.shortcut = nn.Sequential()
27-
if stride != 1 or in_planes != self.expansion*planes:
27+
if stride != 1 or in_planes != self.expansion*group_width:
2828
self.shortcut = nn.Sequential(
29-
nn.Conv2d(in_planes, self.expansion*planes, kernel_size=1, stride=stride, bias=False),
30-
nn.BatchNorm2d(self.expansion*planes)
29+
nn.Conv2d(in_planes, self.expansion*group_width, kernel_size=1, stride=stride, bias=False),
30+
nn.BatchNorm2d(self.expansion*group_width)
3131
)
3232

3333
def forward(self, x):
@@ -40,26 +40,28 @@ def forward(self, x):
4040

4141

4242
class ResNeXt(nn.Module):
43-
def __init__(self, block, num_blocks, cardinality=32, num_classes=10):
43+
def __init__(self, num_blocks, cardinality, bottleneck_width, num_classes=10):
4444
super(ResNeXt, self).__init__()
45+
self.cardinality = cardinality
46+
self.bottleneck_width = bottleneck_width
4547
self.in_planes = 64
4648

4749
self.conv1 = nn.Conv2d(3, 64, kernel_size=1, bias=False)
4850
self.bn1 = nn.BatchNorm2d(64)
49-
self.layer1 = self._make_layer(block, 64, num_blocks[0], 1, cardinality)
50-
self.layer2 = self._make_layer(block, 128, num_blocks[1], 2, cardinality)
51-
self.layer3 = self._make_layer(block, 256, num_blocks[2], 2, cardinality)
52-
# self.layer4 = self._make_layer(block, 512, num_blocks[3], 2, cardinality)
53-
self.linear = nn.Linear(512, num_classes)
51+
self.layer1 = self._make_layer(num_blocks[0], 1)
52+
self.layer2 = self._make_layer(num_blocks[1], 2)
53+
self.layer3 = self._make_layer(num_blocks[2], 2)
54+
# self.layer4 = self._make_layer(num_blocks[3], 2)
55+
self.linear = nn.Linear(cardinality*bottleneck_width*8, num_classes)
5456

55-
self.init_params()
56-
57-
def _make_layer(self, block, planes, num_blocks, stride, cardinality):
57+
def _make_layer(self, num_blocks, stride):
5858
strides = [stride] + [1]*(num_blocks-1)
5959
layers = []
6060
for stride in strides:
61-
layers.append(block(self.in_planes, planes, stride, cardinality))
62-
self.in_planes = planes * block.expansion
61+
layers.append(Block(self.in_planes, self.cardinality, self.bottleneck_width, stride))
62+
self.in_planes = Block.expansion * self.cardinality * self.bottleneck_width
63+
# Increase bottleneck_width by 2 after each stage.
64+
self.bottleneck_width *= 2
6365
return nn.Sequential(*layers)
6466

6567
def forward(self, x):
@@ -73,27 +75,23 @@ def forward(self, x):
7375
out = self.linear(out)
7476
return out
7577

76-
def init_params(self):
77-
'''Init layer parameters.'''
78-
for m in self.modules():
79-
if isinstance(m, nn.Conv2d):
80-
init.kaiming_normal(m.weight, mode='fan_out')
81-
if m.bias:
82-
init.constant(m.bias, 0)
83-
elif isinstance(m, nn.BatchNorm2d):
84-
init.constant(m.weight, 1)
85-
init.constant(m.bias, 0)
86-
elif isinstance(m, nn.Linear):
87-
init.normal(m.weight, std=1e-3)
88-
if m.bias:
89-
init.constant(m.bias, 0)
90-
91-
92-
def ResNeXt29():
93-
return ResNeXt(Block, [3,3,3])
94-
95-
96-
# net = resnext_cifar()
97-
# x = torch.randn(1,3,32,32)
98-
# y = net(Variable(x))
99-
# print(y.size())
78+
79+
def ResNeXt29_2x64d():
80+
return ResNeXt(num_blocks=[3,3,3], cardinality=2, bottleneck_width=64)
81+
82+
def ResNeXt29_4x64d():
83+
return ResNeXt(num_blocks=[3,3,3], cardinality=4, bottleneck_width=64)
84+
85+
def ResNeXt29_8x64d():
86+
return ResNeXt(num_blocks=[3,3,3], cardinality=8, bottleneck_width=64)
87+
88+
def ResNeXt29_32x4d():
89+
return ResNeXt(num_blocks=[3,3,3], cardinality=32, bottleneck_width=4)
90+
91+
def test_resnext():
92+
net = ResNeXt29_2x64d()
93+
x = torch.randn(1,3,32,32)
94+
y = net(Variable(x))
95+
print(y.size())
96+
97+
# test_resnext()

utils.py

Lines changed: 15 additions & 13 deletions
Original file line numberDiff line numberDiff line change
@@ -9,6 +9,7 @@
99
import math
1010

1111
import torch.nn as nn
12+
import torch.nn.init as init
1213

1314

1415
def get_mean_and_std(dataset):
@@ -34,19 +35,20 @@ def get_mean_and_std(dataset):
3435
# print(mean)
3536
# print(std)
3637

37-
38-
def msr_init(net):
39-
'''Initialize layer parameters.'''
40-
for layer in net:
41-
if type(layer) == nn.Conv2d:
42-
n = layer.kernel_size[0]*layer.kernel_size[1]*layer.out_channels
43-
layer.weight.data.normal_(0, math.sqrt(2./n))
44-
layer.bias.data.zero_()
45-
elif type(layer) == nn.BatchNorm2d:
46-
layer.weight.data.fill_(1)
47-
layer.bias.data.zero_()
48-
elif type(layer) == nn.Linear:
49-
layer.bias.data.zero_()
38+
def init_params(net):
39+
'''Init layer parameters.'''
40+
for m in net.modules():
41+
if isinstance(m, nn.Conv2d):
42+
init.kaiming_normal(m.weight, mode='fan_out')
43+
if m.bias:
44+
init.constant(m.bias, 0)
45+
elif isinstance(m, nn.BatchNorm2d):
46+
init.constant(m.weight, 1)
47+
init.constant(m.bias, 0)
48+
elif isinstance(m, nn.Linear):
49+
init.normal(m.weight, std=1e-3)
50+
if m.bias:
51+
init.constant(m.bias, 0)
5052

5153

5254
_, term_width = os.popen('stty size', 'r').read().split()

0 commit comments

Comments
 (0)