import torch
import torchvision
import PIL
import requests
import io
import matplotlib.pyplot as plt
1. imports
2. torch.einsum
A. transpose
-
test tensor
= torch.arange(12).reshape(4,3)
tsr tsr
tensor([[ 0, 1, 2],
[ 3, 4, 5],
[ 6, 7, 8],
[ 9, 10, 11]])
-
행렬을 transpose 하는 방법 1
tsr.t()
tensor([[ 0, 3, 6, 9],
[ 1, 4, 7, 10],
[ 2, 5, 8, 11]])
-
행렬을 transpose 하는 방법 2
'ij -> ji', tsr) torch.einsum(
tensor([[ 0, 3, 6, 9],
[ 1, 4, 7, 10],
[ 2, 5, 8, 11]])
B. 행렬곱
-
test tensors
= torch.arange(12).reshape(4,3).float()
tsr1 = torch.arange(15).reshape(3,5).float()
tsr2 tsr1,tsr2
(tensor([[ 0., 1., 2.],
[ 3., 4., 5.],
[ 6., 7., 8.],
[ 9., 10., 11.]]),
tensor([[ 0., 1., 2., 3., 4.],
[ 5., 6., 7., 8., 9.],
[10., 11., 12., 13., 14.]]))
-
행렬곱을 수행하는 방법1
@ tsr2 tsr1
tensor([[ 25., 28., 31., 34., 37.],
[ 70., 82., 94., 106., 118.],
[115., 136., 157., 178., 199.],
[160., 190., 220., 250., 280.]])
-
행렬곱을 수행하는 방법2
'ij, jk -> ik', tsr1,tsr2) torch.einsum(
tensor([[ 25., 28., 31., 34., 37.],
[ 70., 82., 94., 106., 118.],
[115., 136., 157., 178., 199.],
[160., 190., 220., 250., 280.]])
C. img_plt
vs img_pytorch
-
r,g,b 를 의미하는 tensor
= torch.zeros(16).reshape(4,4) + 1.0
r = torch.zeros(16).reshape(4,4)
g = torch.zeros(16).reshape(4,4) b
-
torch를 쓰기 위해서는 이미지가 이렇게 저장되어 있어야함
= torch.stack([r,g,b],axis=0).reshape(1,3,4,4)
img_pytorch print(img_pytorch)
print(img_pytorch.shape)
tensor([[[[1., 1., 1., 1.],
[1., 1., 1., 1.],
[1., 1., 1., 1.],
[1., 1., 1., 1.]],
[[0., 0., 0., 0.],
[0., 0., 0., 0.],
[0., 0., 0., 0.],
[0., 0., 0., 0.]],
[[0., 0., 0., 0.],
[0., 0., 0., 0.],
[0., 0., 0., 0.],
[0., 0., 0., 0.]]]])
torch.Size([1, 3, 4, 4])
= torch.stack([r,g,b],axis=-1)
img_matplotlib print(img_matplotlib)
print(img_matplotlib.shape)
plt.imshow(img_matplotlib)
tensor([[[1., 0., 0.],
[1., 0., 0.],
[1., 0., 0.],
[1., 0., 0.]],
[[1., 0., 0.],
[1., 0., 0.],
[1., 0., 0.],
[1., 0., 0.]],
[[1., 0., 0.],
[1., 0., 0.],
[1., 0., 0.],
[1., 0., 0.]],
[[1., 0., 0.],
[1., 0., 0.],
[1., 0., 0.],
[1., 0., 0.]]])
torch.Size([4, 4, 3])
# 잘못된코드
4,4,3)) plt.imshow(img_pytorch.squeeze().reshape(
# 올바른코드1
'cij -> ijc', img_pytorch.squeeze())) plt.imshow(torch.einsum(
# 올바른코드2
1,2,0)) plt.imshow(img_pytorch.squeeze().permute(
3. 이미지 자료 처리
A. 데이터
-
데이터 다운로드
torchvision.__version__
'0.20.1'
= torchvision.datasets.OxfordIIITPet(
train_dataset ='./data',
root='trainval',
split=True,
download='binary-category'
target_types
)= torchvision.datasets.OxfordIIITPet(
test_dataset ='./data',
root='test',
split=True,
download='binary-category'
target_types )
-
고양이는 0, 강아지는 1
0][0] train_dataset[
0][1] train_dataset[
0
B. 이미지 변환
-
x_pil
을 tensor로 바꾸어 보자
= train_dataset[0][0]
x_pil x_pil
= torchvision.transforms.ToTensor() # PIL를 텐서로 만드는 함수를 리턴
to_tensor =to_tensor(x_pil)
x_tensorprint(x_tensor.shape)
x_tensor
torch.Size([3, 500, 394])
tensor([[[0.1451, 0.1373, 0.1412, ..., 0.9686, 0.9765, 0.9765],
[0.1373, 0.1373, 0.1451, ..., 0.9647, 0.9725, 0.9765],
[0.1373, 0.1412, 0.1529, ..., 0.9686, 0.9804, 0.9804],
...,
[0.0196, 0.0157, 0.0157, ..., 0.2863, 0.2471, 0.2706],
[0.0157, 0.0118, 0.0118, ..., 0.2392, 0.2157, 0.2510],
[0.1098, 0.1098, 0.1059, ..., 0.2314, 0.2549, 0.2980]],
[[0.0784, 0.0706, 0.0745, ..., 0.9725, 0.9725, 0.9725],
[0.0706, 0.0706, 0.0784, ..., 0.9686, 0.9686, 0.9725],
[0.0706, 0.0745, 0.0863, ..., 0.9647, 0.9765, 0.9765],
...,
[0.0235, 0.0196, 0.0196, ..., 0.4627, 0.4039, 0.4314],
[0.0118, 0.0078, 0.0078, ..., 0.3882, 0.3843, 0.4235],
[0.1059, 0.1059, 0.1059, ..., 0.3686, 0.4157, 0.4588]],
[[0.0471, 0.0392, 0.0431, ..., 0.9922, 0.9922, 0.9922],
[0.0392, 0.0392, 0.0471, ..., 0.9843, 0.9882, 0.9922],
[0.0392, 0.0431, 0.0549, ..., 0.9843, 0.9961, 0.9961],
...,
[0.0941, 0.0902, 0.0902, ..., 0.9608, 0.9216, 0.8784],
[0.0745, 0.0706, 0.0706, ..., 0.9255, 0.9373, 0.8980],
[0.1373, 0.1373, 0.1373, ..., 0.8392, 0.9098, 0.8745]]])
1,2,0)) plt.imshow(to_tensor(x_pil) .permute(
-
궁극적으로는 train_dataset
의 모든 이미지를 (3680,3,???,???)
로 정리하여 X
라고 하고 싶음. \(\to\) 이걸 하기 위해서는 이미지 크기를 통일시켜야함. \(\to\) 이미지크기를 통일시키는 방법을 알아보자.
0][0]).shape to_tensor(train_dataset[
torch.Size([3, 500, 394])
1][0]).shape to_tensor(train_dataset[
torch.Size([3, 313, 450])
-
512,512로 이미지 조정
= torchvision.transforms.Resize((512,512)) # 512,512로 이미지를 조정해주는 함수가 리턴 resize
0][0])).shape to_tensor(resize(train_dataset[
torch.Size([3, 512, 512])
0][0])).permute(1,2,0)) plt.imshow(to_tensor(resize(train_dataset[
-
크기가 8인 이미지들의 배치를 만들기
= torch.stack([to_tensor(resize(train_dataset[n][0])) for n in range(8)],axis=0)
Xm Xm.shape
torch.Size([8, 3, 512, 512])
5. AP Layer
-
채널별로 평균을 구하는 Layer
= torch.nn.AdaptiveAvgPool2d(output_size=1) ap
= torch.arange(1*3*4*4).reshape(1,3,4,4).float()
X X
tensor([[[[ 0., 1., 2., 3.],
[ 4., 5., 6., 7.],
[ 8., 9., 10., 11.],
[12., 13., 14., 15.]],
[[16., 17., 18., 19.],
[20., 21., 22., 23.],
[24., 25., 26., 27.],
[28., 29., 30., 31.]],
[[32., 33., 34., 35.],
[36., 37., 38., 39.],
[40., 41., 42., 43.],
[44., 45., 46., 47.]]]])
# 채널별평균 ap(X)
tensor([[[[ 7.5000]],
[[23.5000]],
[[39.5000]]]])
B.AP, Linear의 교환
= X[:,0,:,:]
r = X[:,1,:,:]
g = X[:,2,:,:] b
*0.1 + ap(g)*0.2 + ap(b)*0.3 ap(r)
tensor([[[17.3000]]])
*0.1 + g*0.2 + b*0.3) ap(r
tensor([[[17.3000]]])
-
위의 두 계산결과를 torch.nn.AdaptiveAvgPool2d
, torch.nn.Linear
, 그리고 torch.nn.Flatten()
을 조합하여 구현해보자.
# ap(r)*0.1 + ap(g)*0.2 + ap(b)*0.3
= torch.nn.AdaptiveAvgPool2d(output_size=1)
ap = torch.nn.Flatten()
flattn = torch.nn.Linear(3,1,bias=False)
linr = torch.tensor([[ 0.1, 0.2, 0.3]])
linr.weight.data #---#
print(X.shape)
print(ap(X).shape) # ap(r), ap(g), ap(b) 의 값이 들어있음.
print(flattn(ap(X)).shape) # [ap(r), ap(g), ap(b)] 형태로
print(linr(flattn(ap(X))).shape) # ap(r)*0.1 + ap(g)*0.2 + ap(b)*0.3
torch.Size([1, 3, 4, 4])
torch.Size([1, 3, 1, 1])
torch.Size([1, 3])
torch.Size([1, 1])
#ap(r*0.1 + g*0.2 + b*0.3)
= torch.nn.AdaptiveAvgPool2d(output_size=1)
ap = torch.nn.Flatten()
flattn = torch.nn.Linear(3,1,bias=False)
linr = torch.tensor([[ 0.1, 0.2, 0.3]])
linr.weight.data #---#
def _linr(X):
return torch.einsum('ocij, kc -> okij', X, linr.weight.data)
#---#
print(X.shape) #
print(_linr(X).shape) # r*0.1 + g*0.2 + b*0.3
print(ap(_linr(X)).shape) # ap(r*0.1 + g*0.2 + b*0.3 )
print(flattn(ap(_linr(X))).shape)
torch.Size([1, 3, 4, 4])
torch.Size([1, 1, 4, 4])
torch.Size([1, 1, 1, 1])
torch.Size([1, 1])
5. CAM(Ahou et al. 2016)의 구현
ref: https://arxiv.org/abs/1512.04150
-
이 강의노트는 위의 논문의 내용을 재구성하였음.
A. 0단계 – (X,y), (XX,yy)
= torchvision.datasets.OxfordIIITPet(
train_dataset ='./data',
root='trainval',
split=True,
download='binary-category',
target_types
)= torchvision.datasets.OxfordIIITPet(
test_dataset ='./data',
root='test',
split=True,
download='binary-category',
target_types )
= torchvision.transforms.Compose([
compose 512,512)),
torchvision.transforms.Resize((
torchvision.transforms.ToTensor() ])
= torch.stack([compose(train_dataset[i][0]) for i in range(3680)],axis=0)
X = torch.stack([compose(test_dataset[i][0]) for i in range(3669)],axis=0)
XX = torch.tensor([train_dataset[i][1] for i in range(3680)]).reshape(-1,1).float()
y = torch.tensor([test_dataset[i][1] for i in range(3669)]).reshape(-1,1).float() yy
B. 1단계 - 이미지 분류 잘하는 네트워크 선택 후 학습
43052)
torch.manual_seed(#--Step1
= torch.utils.data.TensorDataset(X,y)
ds_train = torch.utils.data.DataLoader(ds_train, batch_size=32, shuffle=True)
dl_train = torch.utils.data.TensorDataset(XX,yy)
ds_test = torch.utils.data.DataLoader(ds_test, batch_size=32)
dl_test #--Step2
= torchvision.models.resnet18(pretrained=True)
resnet18 = torch.nn.Linear(512,1)
resnet18.fc = torch.nn.BCEWithLogitsLoss()
loss_fn = torch.optim.Adam(resnet18.parameters(), lr=1e-5)
optimizr #--Step3
"cuda:0")
resnet18.to(for epoc in range(3):
resnet18.train()for Xm,ym in dl_train:
= Xm.to("cuda:0")
Xm = ym.to("cuda:0")
ym #1
= resnet18(Xm)
netout #2
= loss_fn(netout,ym)
loss #3
loss.backward()#4
optimizr.step()
optimizr.zero_grad()#---#
eval()
resnet18.= 0
s for Xm,ym in dl_train:
= Xm.to("cuda:0")
Xm = ym.to("cuda:0")
ym = s + ((resnet18(Xm).data > 0) == ym).sum().item()
s = s/3680
acc print(f"train_acc = {acc:.4f}")
#--Step4
eval()
resnet18.= 0
s for Xm,ym in dl_test:
= Xm.to("cuda:0")
Xm = ym.to("cuda:0")
ym = s + ((resnet18(Xm).data > 0) == ym).sum().item()
s = s/3669
acc print(f"test_acc = {acc:.4f}")
C. 2단계 - Linear 와 AP의 순서를 바꿈
-
resnet18
을 재구성하여 net
을 만들자
resnet18._forward_impl??
Signature: resnet18._forward_impl(x: torch.Tensor) -> torch.Tensor Docstring: <no docstring> Source: def _forward_impl(self, x: Tensor) -> Tensor: # See note [TorchScript super()] x = self.conv1(x) x = self.bn1(x) x = self.relu(x) x = self.maxpool(x) x = self.layer1(x) x = self.layer2(x) x = self.layer3(x) x = self.layer4(x) x = self.avgpool(x) x = torch.flatten(x, 1) x = self.fc(x) return x File: ~/anaconda3/envs/test/lib/python3.12/site-packages/torchvision/models/resnet.py Type: method
= torch.nn.Sequential(
stem
torch.nn.Sequential(
resnet18.conv1,
resnet18.bn1,
resnet18.relu,
resnet18.maxpool
),
resnet18.layer1,
resnet18.layer2,
resnet18.layer3,
resnet18.layer4
)= torch.nn.Sequential(
head
resnet18.avgpool,
torch.nn.Flatten(),
resnet18.fc
)= torch.nn.Sequential(
net
stem,
head )
-
아직은 resnet18
과 똑같은 기능, 인덱스로 접근 가능!
-
1개의 observation을 고정
= X[[0]].to("cuda:0") x
-
이렇게 해도 될듯
0]],axis=0).shape torch.stack([X[
torch.Size([1, 3, 512, 512])
net(x), resnet18(x)
(tensor([[-5.5534]], device='cuda:0', grad_fn=<AddmmBackward0>),
tensor([[-5.5534]], device='cuda:0', grad_fn=<AddmmBackward0>))
-
위와 같은 값 -5.5534이 나오는 과정을 추적하여 보자.
# 계산방식1: 원래계산방식
= head[0]
ap = head[1]
flattn = head[2]
linr #---#
print(f"{x.shape} -- x")
print(f"{stem(x).shape} -- stem(x)")
print(f"{ap(stem(x)).shape} -- ap(stem(x))")
print(f"{flattn(ap(stem(x))).shape} -- flattn(ap(stem(x)))")
print(f"{linr(flattn(ap(stem(x)))).shape} -- linr(flattn(ap(stem(x))))")
torch.Size([1, 3, 512, 512]) -- x
torch.Size([1, 512, 16, 16]) -- stem(x)
torch.Size([1, 512, 1, 1]) -- ap(stem(x))
torch.Size([1, 512]) -- flattn(ap(stem(x)))
torch.Size([1, 1]) -- linr(flattn(ap(stem(x))))
현재 네트워크 \[\underset{(1,3,512,512)}{\boldsymbol x} \overset{stem}{\to} \left( \underset{(1,512,16,16)}{\tilde{\boldsymbol x}} \overset{ap}{\to} \underset{(1,512,1,1)}{{\boldsymbol \sharp}}\overset{flattn}{\to} \underset{(1,512)}{{\boldsymbol \sharp}}\overset{linr}{\to} \underset{(1,1)}{logit}\right) = [[-5.5613]]\]
바꾸고 싶은 네트워크 \[\underset{(1,3,224,224)}{\boldsymbol x} \overset{stem}{\to} \left( \underset{(1,512,16,16)}{\tilde{\boldsymbol x}} \overset{\_linr}{\to} \underset{(1,1,16,16)}{{\boldsymbol \sharp}}\overset{ap}{\to} \underset{(1,1,1,1)}{{\boldsymbol \sharp}}\overset{flattn}{\to} \underset{(1,1)}{logit}\right) = [[-5.5613]]\]
# 계산방식2
= head[0]
ap = head[1]
flattn = head[2]
linr def _linr(xtilde):
return torch.einsum('ocij, kc -> okij', xtilde, linr.weight.data) + linr.bias.data
#---#
print(f"{x.shape} -- x")
print(f"{stem(x).shape} -- stem(x)")
print(f"{_linr(stem(x)).shape} -- _linr(stem(x))")
print(f"{ap(_linr(stem(x))).shape} -- ap(_linr(stem(x)))")
print(f"{flattn(ap(_linr(stem(x)))).shape} -- flattn(ap(_linr(stem(x))))")
torch.Size([1, 3, 512, 512]) -- x
torch.Size([1, 512, 16, 16]) -- stem(x)
torch.Size([1, 1, 16, 16]) -- _linr(stem(x))
torch.Size([1, 1, 1, 1]) -- ap(_linr(stem(x)))
torch.Size([1, 1]) -- flattn(ap(_linr(stem(x))))
linr(flattn(ap(stem(x)))), flattn(ap(_linr(stem(x))))
(tensor([[-5.5534]], device='cuda:0', grad_fn=<AddmmBackward0>),
tensor([[-5.5534]], device='cuda:0', grad_fn=<ViewBackward0>))
-
멈추고 생각해보기…!!
-
원래 계산방식을 적용
linr(flattn(ap(stem(x))))
tensor([[-5.5534]], device='cuda:0', grad_fn=<AddmmBackward0>)
-
바뀐 계산방식을 적용
flattn(ap(_linr(stem(x))))
tensor([[-5.5534]], device='cuda:0', grad_fn=<ViewBackward0>)
-
바뀐 계산방식을 좀더 파고 들어서 분석해보자.
long() _linr(stem(x)).
tensor([[[[ 0, 0, 0, 0, 0, 0, 0, 0, -1, -2, -2, -2, 0, 0,
0, 0],
[ 0, 0, 0, 0, 0, 0, 0, -1, 0, -2, -3, -4, -1, 0,
0, 0],
[ 0, 0, 0, 0, 0, 0, 0, 0, 0, 0, -3, -5, -4, -2,
0, 0],
[ 0, 0, 0, 0, 0, 0, 0, 0, -1, -4, -9, -12, -12, -8,
-2, 0],
[ 0, 0, 0, 0, 0, 0, -1, -5, -11, -16, -20, -24, -22, -14,
-4, 0],
[ -1, -1, 0, 0, 0, 0, -5, -16, -28, -35, -40, -42, -37, -23,
-7, 0],
[ -1, -1, 0, 0, 0, 0, -10, -28, -47, -56, -56, -52, -42, -25,
-7, 0],
[ 0, -1, 0, 0, 0, 0, -11, -29, -49, -57, -54, -47, -34, -19,
-4, 1],
[ 0, 0, 0, 0, 0, 0, -8, -21, -36, -42, -38, -29, -18, -8,
-1, 0],
[ 0, 0, -1, -1, 0, 0, -3, -9, -16, -19, -16, -11, -4, 0,
0, 0],
[ 0, 0, -2, -2, 0, 0, -1, -3, -6, -6, -5, -2, 0, 0,
1, 1],
[ 0, 0, -2, -1, 0, 0, -1, -3, -4, -3, -1, 0, 0, 0,
1, 1],
[ 0, 0, -1, 0, 0, 0, 0, -2, -2, -1, 0, 0, 0, 0,
1, 1],
[ 0, 0, 0, 0, 0, 0, 0, -1, -1, 0, 0, 0, 0, 1,
2, 1],
[ -1, -1, -1, -1, 0, 0, 0, -1, 0, 0, 0, 1, 1, 2,
2, 2],
[ 0, -1, -1, 0, 0, 0, 0, 0, 0, 0, 0, 1, 1, 1,
1, 1]]]], device='cuda:0')
- 여러가지 값들이 있지만 (16*16=256개의 값들) 아무튼 이 값들의 평균은 -5.5613 임. (이 값이 작을수록 이 그림은 고양이라는 의미임)
- 그런데 살펴보니까 대부분의 위치에서 -에 가까운 값을가지고, 특정위치에서만 엄청 작은 값이 존재하여 -5.5613 이라는 평균값이 나오는 것임.
- 결국 이 특정위치에 존재하는 엄청 작은 값들이
x
가 고양이라고 판단하는 근거가 된다.
바꾸고 싶은 네트워크 – why라는 이름을 적용하여 \[\underset{(1,3,224,224)}{\boldsymbol x} \overset{stem}{\to} \left( \underset{(1,512,16,16)}{\tilde{\boldsymbol x}} \overset{\_linr}{\to} \underset{(1,1,16,16)}{\bf why}\overset{ap}{\to} \underset{(1,1,1,1)}{{\boldsymbol \sharp}}\overset{flattn}{\to} \underset{(1,1)}{logit}\right) = [[-5.5613]]\]
= _linr(stem(x)) why
D. 3단계 -WHY 시각화
-
시각화1 - why와 img를 겹쳐서 그려보자
="bwr")
plt.imshow(why.squeeze().cpu().detach(),cmap plt.colorbar()
1,2,0))
plt.imshow(x.cpu().detach().squeeze().permute(= torch.nn.functional.interpolate(
why_resized
why,=(512,512),
size="bilinear",
mode
)="bwr",alpha=0.5)
plt.imshow(why_resized.squeeze().cpu().detach(),cmap plt.colorbar()
-
시각화2 - colormap을 magma
로 적용
= X[[0]].to("cuda:0")
x if net(x) > 0:
= "dog"
pred = _linr(stem(x))
why else:
= "cat"
pred = - _linr(stem(x))
why 1,2,0))
plt.imshow(x.cpu().detach().squeeze().permute(= torch.nn.functional.interpolate(
why_resized
why,=(512,512),
size="bilinear",
mode
)="magma",alpha=0.5)
plt.imshow(why_resized.squeeze().cpu().detach(),cmapf"prediction = {pred}"); plt.title(
-
시각화3 - 하니를 시각화
= 'https://github.com/guebin/DL2025/blob/main/imgs/hani1.jpeg?raw=true'
url = PIL.Image.open(
hani_pil
io.BytesIO(requests.get(url).content) )
= compose(hani_pil).reshape(1,3,512,512).to("cuda:0")
x if net(x) > 0:
= "dog"
pred = _linr(stem(x))
why else:
= "cat"
pred = - _linr(stem(x))
why 1,2,0))
plt.imshow(x.cpu().detach().squeeze().permute(= torch.nn.functional.interpolate(
why_resized
why,=(512,512),
size="bilinear",
mode
)="magma",alpha=0.5)
plt.imshow(why_resized.squeeze().cpu().detach(),cmapf"prediction = {pred}"); plt.title(
-
시각화4 - XX
의 이미지를 시각화해보자
= plt.subplots(5,5)
fig, ax #---#
= 0
k for i in range(5):
for j in range(5):
= XX[[k]].to("cuda:0")
x if net(x) > 0:
= "dog"
pred = _linr(stem(x))
why else:
= "cat"
pred = - _linr(stem(x))
why 1,2,0))
plt.imshow(x.cpu().detach().squeeze().permute(= torch.nn.functional.interpolate(
why_resized
why,=(512,512),
size="bilinear",
mode
)1,2,0).cpu())
ax[i][j].imshow(x.squeeze().permute(="magma",alpha=0.5)
ax[i][j].imshow(why_resized.squeeze().cpu().detach(),cmapf"prediction = {pred}");
ax[i][j].set_title(
ax[i][j].set_xticks([])
ax[i][j].set_yticks([])= k+50
k 16)
fig.set_figheight(16)
fig.set_figwidth( fig.tight_layout()
= plt.subplots(5,5)
fig, ax #---#
= 1
k for i in range(5):
for j in range(5):
= XX[[k]].to("cuda:0")
x if net(x) > 0:
= "dog"
pred = _linr(stem(x))
why else:
= "cat"
pred = - _linr(stem(x))
why 1,2,0))
plt.imshow(x.cpu().detach().squeeze().permute(= torch.nn.functional.interpolate(
why_resized
why,=(512,512),
size="bilinear",
mode
)1,2,0).cpu())
ax[i][j].imshow(x.squeeze().permute(="magma",alpha=0.5)
ax[i][j].imshow(why_resized.squeeze().cpu().detach(),cmapf"prediction = {pred}");
ax[i][j].set_title(
ax[i][j].set_xticks([])
ax[i][j].set_yticks([])= k+50
k 16)
fig.set_figheight(16)
fig.set_figwidth( fig.tight_layout()
= plt.subplots(5,5)
fig, ax #---#
= 2
k for i in range(5):
for j in range(5):
= XX[[k]].to("cuda:0")
x if net(x) > 0:
= "dog"
pred = _linr(stem(x))
why else:
= "cat"
pred = - _linr(stem(x))
why 1,2,0))
plt.imshow(x.cpu().detach().squeeze().permute(= torch.nn.functional.interpolate(
why_resized
why,=(512,512),
size="bilinear",
mode
)1,2,0).cpu())
ax[i][j].imshow(x.squeeze().permute(="magma",alpha=0.5)
ax[i][j].imshow(why_resized.squeeze().cpu().detach(),cmapf"prediction = {pred}");
ax[i][j].set_title(
ax[i][j].set_xticks([])
ax[i][j].set_yticks([])= k+50
k 16)
fig.set_figheight(16)
fig.set_figwidth( fig.tight_layout()
= plt.subplots(5,5)
fig, ax #---#
= 3
k for i in range(5):
for j in range(5):
= XX[[k]].to("cuda:0")
x if net(x) > 0:
= "dog"
pred = _linr(stem(x))
why else:
= "cat"
pred = - _linr(stem(x))
why 1,2,0))
plt.imshow(x.cpu().detach().squeeze().permute(= torch.nn.functional.interpolate(
why_resized
why,=(512,512),
size="bilinear",
mode
)1,2,0).cpu())
ax[i][j].imshow(x.squeeze().permute(="magma",alpha=0.5)
ax[i][j].imshow(why_resized.squeeze().cpu().detach(),cmapf"prediction = {pred}");
ax[i][j].set_title(
ax[i][j].set_xticks([])
ax[i][j].set_yticks([])= k+50
k 16)
fig.set_figheight(16)
fig.set_figwidth( fig.tight_layout()
6. CAM의 한계
-
구조적 제약이 있음
- head-part가 ap와 linr로만 구성되어야 가능
- 그렇지 않은 네트워크는 임의로 재구성하여 head-part를 ap와 linr로많 구성해야함(CAM을 적용하기 위한 구조로 만들기 위해)
-
이후로 등장한 grad-cam
은 구조의 제약 없이 거의 모든 CNN에 적용 가능