123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261262263264265266267268269270271272273274275276277278279280281282283284285286287288289290291292293294295296297298299300301302303304305306307308309310311312313314315316317318319320321322323324325326327328329330 |
- # Copyright 2016 Google Inc. All Rights Reserved.
- #
- # Licensed under the Apache License, Version 2.0 (the "License");
- # you may not use this file except in compliance with the License.
- # You may obtain a copy of the License at
- #
- # http://www.apache.org/licenses/LICENSE-2.0
- #
- # Unless required by applicable law or agreed to in writing, software
- # distributed under the License is distributed on an "AS IS" BASIS,
- # WITHOUT WARRANTIES OR CONDITIONS OF ANY KIND, either express or implied.
- # See the License for the specific language governing permissions and
- # limitations under the License.
- # ==============================================================================
- """Inception-v3 expressed in TensorFlow-Slim.
- Usage:
- # Parameters for BatchNorm.
- batch_norm_params = {
- # Decay for the batch_norm moving averages.
- 'decay': BATCHNORM_MOVING_AVERAGE_DECAY,
- # epsilon to prevent 0s in variance.
- 'epsilon': 0.001,
- }
- # Set weight_decay for weights in Conv and FC layers.
- with slim.arg_scope([slim.ops.conv2d, slim.ops.fc], weight_decay=0.00004):
- with slim.arg_scope([slim.ops.conv2d],
- stddev=0.1,
- activation=tf.nn.relu,
- batch_norm_params=batch_norm_params):
- # Force all Variables to reside on the CPU.
- with slim.arg_scope([slim.variables.variable], device='/cpu:0'):
- logits, endpoints = slim.inception.inception_v3(
- images,
- dropout_keep_prob=0.8,
- num_classes=num_classes,
- is_training=for_training,
- restore_logits=restore_logits,
- scope=scope)
- """
- from __future__ import absolute_import
- from __future__ import division
- from __future__ import print_function
- import tensorflow as tf
- from inception.slim import ops
- from inception.slim import scopes
- def inception_v3(inputs,
- dropout_keep_prob=0.8,
- num_classes=1000,
- is_training=True,
- restore_logits=True,
- scope=''):
- """Latest Inception from http://arxiv.org/abs/1512.00567.
- "Rethinking the Inception Architecture for Computer Vision"
- Christian Szegedy, Vincent Vanhoucke, Sergey Ioffe, Jonathon Shlens,
- Zbigniew Wojna
- Args:
- inputs: a tensor of size [batch_size, height, width, channels].
- dropout_keep_prob: dropout keep_prob.
- num_classes: number of predicted classes.
- is_training: whether is training or not.
- restore_logits: whether or not the logits layers should be restored.
- Useful for fine-tuning a model with different num_classes.
- scope: Optional scope for op_scope.
- Returns:
- a list containing 'logits', 'aux_logits' Tensors.
- """
- # end_points will collect relevant activations for external use, for example
- # summaries or losses.
- end_points = {}
- with tf.op_scope([inputs], scope, 'inception_v3'):
- with scopes.arg_scope([ops.conv2d, ops.fc, ops.batch_norm, ops.dropout],
- is_training=is_training):
- with scopes.arg_scope([ops.conv2d, ops.max_pool, ops.avg_pool],
- stride=1, padding='VALID'):
- # 299 x 299 x 3
- end_points['conv0'] = ops.conv2d(inputs, 32, [3, 3], stride=2,
- scope='conv0')
- # 149 x 149 x 32
- end_points['conv1'] = ops.conv2d(end_points['conv0'], 32, [3, 3],
- scope='conv1')
- # 147 x 147 x 32
- end_points['conv2'] = ops.conv2d(end_points['conv1'], 64, [3, 3],
- padding='SAME', scope='conv2')
- # 147 x 147 x 64
- end_points['pool1'] = ops.max_pool(end_points['conv2'], [3, 3],
- stride=2, scope='pool1')
- # 73 x 73 x 64
- end_points['conv3'] = ops.conv2d(end_points['pool1'], 80, [1, 1],
- scope='conv3')
- # 71 x 71 x 80.
- end_points['conv4'] = ops.conv2d(end_points['conv3'], 192, [3, 3],
- scope='conv4')
- # 69 x 69 x 192.
- end_points['pool2'] = ops.max_pool(end_points['conv4'], [3, 3],
- stride=2, scope='pool2')
- # 35 x 35 x 192.
- net = end_points['pool2']
- # Inception blocks
- with scopes.arg_scope([ops.conv2d, ops.max_pool, ops.avg_pool],
- stride=1, padding='SAME'):
- # mixed: 35 x 35 x 256.
- with tf.variable_scope('mixed_35x35x256a'):
- with tf.variable_scope('branch1x1'):
- branch1x1 = ops.conv2d(net, 64, [1, 1])
- with tf.variable_scope('branch5x5'):
- branch5x5 = ops.conv2d(net, 48, [1, 1])
- branch5x5 = ops.conv2d(branch5x5, 64, [5, 5])
- with tf.variable_scope('branch3x3dbl'):
- branch3x3dbl = ops.conv2d(net, 64, [1, 1])
- branch3x3dbl = ops.conv2d(branch3x3dbl, 96, [3, 3])
- branch3x3dbl = ops.conv2d(branch3x3dbl, 96, [3, 3])
- with tf.variable_scope('branch_pool'):
- branch_pool = ops.avg_pool(net, [3, 3])
- branch_pool = ops.conv2d(branch_pool, 32, [1, 1])
- net = tf.concat(3, [branch1x1, branch5x5, branch3x3dbl, branch_pool])
- end_points['mixed_35x35x256a'] = net
- # mixed_1: 35 x 35 x 288.
- with tf.variable_scope('mixed_35x35x288a'):
- with tf.variable_scope('branch1x1'):
- branch1x1 = ops.conv2d(net, 64, [1, 1])
- with tf.variable_scope('branch5x5'):
- branch5x5 = ops.conv2d(net, 48, [1, 1])
- branch5x5 = ops.conv2d(branch5x5, 64, [5, 5])
- with tf.variable_scope('branch3x3dbl'):
- branch3x3dbl = ops.conv2d(net, 64, [1, 1])
- branch3x3dbl = ops.conv2d(branch3x3dbl, 96, [3, 3])
- branch3x3dbl = ops.conv2d(branch3x3dbl, 96, [3, 3])
- with tf.variable_scope('branch_pool'):
- branch_pool = ops.avg_pool(net, [3, 3])
- branch_pool = ops.conv2d(branch_pool, 64, [1, 1])
- net = tf.concat(3, [branch1x1, branch5x5, branch3x3dbl, branch_pool])
- end_points['mixed_35x35x288a'] = net
- # mixed_2: 35 x 35 x 288.
- with tf.variable_scope('mixed_35x35x288b'):
- with tf.variable_scope('branch1x1'):
- branch1x1 = ops.conv2d(net, 64, [1, 1])
- with tf.variable_scope('branch5x5'):
- branch5x5 = ops.conv2d(net, 48, [1, 1])
- branch5x5 = ops.conv2d(branch5x5, 64, [5, 5])
- with tf.variable_scope('branch3x3dbl'):
- branch3x3dbl = ops.conv2d(net, 64, [1, 1])
- branch3x3dbl = ops.conv2d(branch3x3dbl, 96, [3, 3])
- branch3x3dbl = ops.conv2d(branch3x3dbl, 96, [3, 3])
- with tf.variable_scope('branch_pool'):
- branch_pool = ops.avg_pool(net, [3, 3])
- branch_pool = ops.conv2d(branch_pool, 64, [1, 1])
- net = tf.concat(3, [branch1x1, branch5x5, branch3x3dbl, branch_pool])
- end_points['mixed_35x35x288b'] = net
- # mixed_3: 17 x 17 x 768.
- with tf.variable_scope('mixed_17x17x768a'):
- with tf.variable_scope('branch3x3'):
- branch3x3 = ops.conv2d(net, 384, [3, 3], stride=2, padding='VALID')
- with tf.variable_scope('branch3x3dbl'):
- branch3x3dbl = ops.conv2d(net, 64, [1, 1])
- branch3x3dbl = ops.conv2d(branch3x3dbl, 96, [3, 3])
- branch3x3dbl = ops.conv2d(branch3x3dbl, 96, [3, 3],
- stride=2, padding='VALID')
- with tf.variable_scope('branch_pool'):
- branch_pool = ops.max_pool(net, [3, 3], stride=2, padding='VALID')
- net = tf.concat(3, [branch3x3, branch3x3dbl, branch_pool])
- end_points['mixed_17x17x768a'] = net
- # mixed4: 17 x 17 x 768.
- with tf.variable_scope('mixed_17x17x768b'):
- with tf.variable_scope('branch1x1'):
- branch1x1 = ops.conv2d(net, 192, [1, 1])
- with tf.variable_scope('branch7x7'):
- branch7x7 = ops.conv2d(net, 128, [1, 1])
- branch7x7 = ops.conv2d(branch7x7, 128, [1, 7])
- branch7x7 = ops.conv2d(branch7x7, 192, [7, 1])
- with tf.variable_scope('branch7x7dbl'):
- branch7x7dbl = ops.conv2d(net, 128, [1, 1])
- branch7x7dbl = ops.conv2d(branch7x7dbl, 128, [7, 1])
- branch7x7dbl = ops.conv2d(branch7x7dbl, 128, [1, 7])
- branch7x7dbl = ops.conv2d(branch7x7dbl, 128, [7, 1])
- branch7x7dbl = ops.conv2d(branch7x7dbl, 192, [1, 7])
- with tf.variable_scope('branch_pool'):
- branch_pool = ops.avg_pool(net, [3, 3])
- branch_pool = ops.conv2d(branch_pool, 192, [1, 1])
- net = tf.concat(3, [branch1x1, branch7x7, branch7x7dbl, branch_pool])
- end_points['mixed_17x17x768b'] = net
- # mixed_5: 17 x 17 x 768.
- with tf.variable_scope('mixed_17x17x768c'):
- with tf.variable_scope('branch1x1'):
- branch1x1 = ops.conv2d(net, 192, [1, 1])
- with tf.variable_scope('branch7x7'):
- branch7x7 = ops.conv2d(net, 160, [1, 1])
- branch7x7 = ops.conv2d(branch7x7, 160, [1, 7])
- branch7x7 = ops.conv2d(branch7x7, 192, [7, 1])
- with tf.variable_scope('branch7x7dbl'):
- branch7x7dbl = ops.conv2d(net, 160, [1, 1])
- branch7x7dbl = ops.conv2d(branch7x7dbl, 160, [7, 1])
- branch7x7dbl = ops.conv2d(branch7x7dbl, 160, [1, 7])
- branch7x7dbl = ops.conv2d(branch7x7dbl, 160, [7, 1])
- branch7x7dbl = ops.conv2d(branch7x7dbl, 192, [1, 7])
- with tf.variable_scope('branch_pool'):
- branch_pool = ops.avg_pool(net, [3, 3])
- branch_pool = ops.conv2d(branch_pool, 192, [1, 1])
- net = tf.concat(3, [branch1x1, branch7x7, branch7x7dbl, branch_pool])
- end_points['mixed_17x17x768c'] = net
- # mixed_6: 17 x 17 x 768.
- with tf.variable_scope('mixed_17x17x768d'):
- with tf.variable_scope('branch1x1'):
- branch1x1 = ops.conv2d(net, 192, [1, 1])
- with tf.variable_scope('branch7x7'):
- branch7x7 = ops.conv2d(net, 160, [1, 1])
- branch7x7 = ops.conv2d(branch7x7, 160, [1, 7])
- branch7x7 = ops.conv2d(branch7x7, 192, [7, 1])
- with tf.variable_scope('branch7x7dbl'):
- branch7x7dbl = ops.conv2d(net, 160, [1, 1])
- branch7x7dbl = ops.conv2d(branch7x7dbl, 160, [7, 1])
- branch7x7dbl = ops.conv2d(branch7x7dbl, 160, [1, 7])
- branch7x7dbl = ops.conv2d(branch7x7dbl, 160, [7, 1])
- branch7x7dbl = ops.conv2d(branch7x7dbl, 192, [1, 7])
- with tf.variable_scope('branch_pool'):
- branch_pool = ops.avg_pool(net, [3, 3])
- branch_pool = ops.conv2d(branch_pool, 192, [1, 1])
- net = tf.concat(3, [branch1x1, branch7x7, branch7x7dbl, branch_pool])
- end_points['mixed_17x17x768d'] = net
- # mixed_7: 17 x 17 x 768.
- with tf.variable_scope('mixed_17x17x768e'):
- with tf.variable_scope('branch1x1'):
- branch1x1 = ops.conv2d(net, 192, [1, 1])
- with tf.variable_scope('branch7x7'):
- branch7x7 = ops.conv2d(net, 192, [1, 1])
- branch7x7 = ops.conv2d(branch7x7, 192, [1, 7])
- branch7x7 = ops.conv2d(branch7x7, 192, [7, 1])
- with tf.variable_scope('branch7x7dbl'):
- branch7x7dbl = ops.conv2d(net, 192, [1, 1])
- branch7x7dbl = ops.conv2d(branch7x7dbl, 192, [7, 1])
- branch7x7dbl = ops.conv2d(branch7x7dbl, 192, [1, 7])
- branch7x7dbl = ops.conv2d(branch7x7dbl, 192, [7, 1])
- branch7x7dbl = ops.conv2d(branch7x7dbl, 192, [1, 7])
- with tf.variable_scope('branch_pool'):
- branch_pool = ops.avg_pool(net, [3, 3])
- branch_pool = ops.conv2d(branch_pool, 192, [1, 1])
- net = tf.concat(3, [branch1x1, branch7x7, branch7x7dbl, branch_pool])
- end_points['mixed_17x17x768e'] = net
- # Auxiliary Head logits
- aux_logits = tf.identity(end_points['mixed_17x17x768e'])
- with tf.variable_scope('aux_logits'):
- aux_logits = ops.avg_pool(aux_logits, [5, 5], stride=3,
- padding='VALID')
- aux_logits = ops.conv2d(aux_logits, 128, [1, 1], scope='proj')
- # Shape of feature map before the final layer.
- shape = aux_logits.get_shape()
- aux_logits = ops.conv2d(aux_logits, 768, shape[1:3], stddev=0.01,
- padding='VALID')
- aux_logits = ops.flatten(aux_logits)
- aux_logits = ops.fc(aux_logits, num_classes, activation=None,
- stddev=0.001, restore=restore_logits)
- end_points['aux_logits'] = aux_logits
- # mixed_8: 17 x 17 x 1280.
- with tf.variable_scope('mixed_17x17x1280a'):
- with tf.variable_scope('branch3x3'):
- branch3x3 = ops.conv2d(net, 192, [1, 1])
- branch3x3 = ops.conv2d(branch3x3, 320, [3, 3], stride=2,
- padding='VALID')
- with tf.variable_scope('branch7x7x3'):
- branch7x7x3 = ops.conv2d(net, 192, [1, 1])
- branch7x7x3 = ops.conv2d(branch7x7x3, 192, [1, 7])
- branch7x7x3 = ops.conv2d(branch7x7x3, 192, [7, 1])
- branch7x7x3 = ops.conv2d(branch7x7x3, 192, [3, 3],
- stride=2, padding='VALID')
- with tf.variable_scope('branch_pool'):
- branch_pool = ops.max_pool(net, [3, 3], stride=2, padding='VALID')
- net = tf.concat(3, [branch3x3, branch7x7x3, branch_pool])
- end_points['mixed_17x17x1280a'] = net
- # mixed_9: 8 x 8 x 2048.
- with tf.variable_scope('mixed_8x8x2048a'):
- with tf.variable_scope('branch1x1'):
- branch1x1 = ops.conv2d(net, 320, [1, 1])
- with tf.variable_scope('branch3x3'):
- branch3x3 = ops.conv2d(net, 384, [1, 1])
- branch3x3 = tf.concat(3, [ops.conv2d(branch3x3, 384, [1, 3]),
- ops.conv2d(branch3x3, 384, [3, 1])])
- with tf.variable_scope('branch3x3dbl'):
- branch3x3dbl = ops.conv2d(net, 448, [1, 1])
- branch3x3dbl = ops.conv2d(branch3x3dbl, 384, [3, 3])
- branch3x3dbl = tf.concat(3, [ops.conv2d(branch3x3dbl, 384, [1, 3]),
- ops.conv2d(branch3x3dbl, 384, [3, 1])])
- with tf.variable_scope('branch_pool'):
- branch_pool = ops.avg_pool(net, [3, 3])
- branch_pool = ops.conv2d(branch_pool, 192, [1, 1])
- net = tf.concat(3, [branch1x1, branch3x3, branch3x3dbl, branch_pool])
- end_points['mixed_8x8x2048a'] = net
- # mixed_10: 8 x 8 x 2048.
- with tf.variable_scope('mixed_8x8x2048b'):
- with tf.variable_scope('branch1x1'):
- branch1x1 = ops.conv2d(net, 320, [1, 1])
- with tf.variable_scope('branch3x3'):
- branch3x3 = ops.conv2d(net, 384, [1, 1])
- branch3x3 = tf.concat(3, [ops.conv2d(branch3x3, 384, [1, 3]),
- ops.conv2d(branch3x3, 384, [3, 1])])
- with tf.variable_scope('branch3x3dbl'):
- branch3x3dbl = ops.conv2d(net, 448, [1, 1])
- branch3x3dbl = ops.conv2d(branch3x3dbl, 384, [3, 3])
- branch3x3dbl = tf.concat(3, [ops.conv2d(branch3x3dbl, 384, [1, 3]),
- ops.conv2d(branch3x3dbl, 384, [3, 1])])
- with tf.variable_scope('branch_pool'):
- branch_pool = ops.avg_pool(net, [3, 3])
- branch_pool = ops.conv2d(branch_pool, 192, [1, 1])
- net = tf.concat(3, [branch1x1, branch3x3, branch3x3dbl, branch_pool])
- end_points['mixed_8x8x2048b'] = net
- # Final pooling and prediction
- with tf.variable_scope('logits'):
- shape = net.get_shape()
- net = ops.avg_pool(net, shape[1:3], padding='VALID', scope='pool')
- # 1 x 1 x 2048
- net = ops.dropout(net, dropout_keep_prob, scope='dropout')
- net = ops.flatten(net, scope='flatten')
- # 2048
- logits = ops.fc(net, num_classes, activation=None, scope='logits',
- restore=restore_logits)
- # 1000
- end_points['logits'] = logits
- end_points['predictions'] = tf.nn.softmax(logits, name='predictions')
- return logits, end_points
|