Skip to content

Commit

Permalink
Updated LFW evaluation code
Browse files Browse the repository at this point in the history
  • Loading branch information
davidsandberg committed Apr 7, 2018
1 parent 113966c commit 9852362
Show file tree
Hide file tree
Showing 6 changed files with 438 additions and 93 deletions.
6 changes: 6 additions & 0 deletions data/learning_rate_schedule_classifier_vggface2.txt
Original file line number Diff line number Diff line change
@@ -0,0 +1,6 @@
# Learning rate schedule
# Maps an epoch number to a learning rate
0: 0.05
100: 0.005
200: 0.0005
276: -1
63 changes: 13 additions & 50 deletions src/facenet.py
Original file line number Diff line number Diff line change
Expand Up @@ -90,62 +90,18 @@ def shuffle_examples(image_paths, labels):
image_paths_shuff, labels_shuff = zip(*shuffle_list)
return image_paths_shuff, labels_shuff

def read_images_from_disk(input_queue):
"""Consumes a single filename and label as a ' '-delimited string.
Args:
filename_and_label_tensor: A scalar string tensor.
Returns:
Two tensors: the decoded image, and the string label.
"""
label = input_queue[1]
file_contents = tf.read_file(input_queue[0])
example = tf.image.decode_image(file_contents, channels=3)
return example, label

def random_rotate_image(image):
angle = np.random.uniform(low=-10.0, high=10.0)
return misc.imrotate(image, angle, 'bicubic')

# def read_and_augment_data(image_list, label_list, image_size, batch_size, max_nrof_epochs,
# random_crop, random_flip, random_rotate, nrof_preprocess_threads, shuffle=True):
#
# images = ops.convert_to_tensor(image_list, dtype=tf.string)
# labels = ops.convert_to_tensor(label_list, dtype=tf.int32)
#
# # Makes an input queue
# input_queue = tf.train.slice_input_producer([images, labels],
# num_epochs=max_nrof_epochs, shuffle=shuffle)
#
# images_and_labels = []
# for _ in range(nrof_preprocess_threads):
# image, label = read_images_from_disk(input_queue)
# if random_rotate:
# image = tf.py_func(random_rotate_image, [image], tf.uint8)
# if random_crop:
# image = tf.random_crop(image, [image_size, image_size, 3])
# else:
# image = tf.image.resize_image_with_crop_or_pad(image, image_size, image_size)
# if random_flip:
# image = tf.image.random_flip_left_right(image)
# #pylint: disable=no-member
# image.set_shape((image_size, image_size, 3))
# image = tf.image.per_image_standardization(image)
# images_and_labels.append([image, label])
#
# image_batch, label_batch = tf.train.batch_join(
# images_and_labels, batch_size=batch_size,
# capacity=4 * nrof_preprocess_threads * batch_size,
# allow_smaller_final_batch=True)
#
# return image_batch, label_batch

# 1: Random rotate 2: Random crop 4: Random flip 8: Fixed image standardization 16: Flip
RANDOM_ROTATE = 1
RANDOM_CROP = 2
RANDOM_FLIP = 4
FIXED_STANDARDIZATION = 8
FLIP = 16
def create_input_pipeline(images_and_labels_list, input_queue, image_size, nrof_preprocess_threads):
def create_input_pipeline(input_queue, image_size, nrof_preprocess_threads, batch_size_placeholder):
images_and_labels_list = []
for _ in range(nrof_preprocess_threads):
filenames, label, control = input_queue.dequeue()
images = []
Expand All @@ -171,7 +127,14 @@ def create_input_pipeline(images_and_labels_list, input_queue, image_size, nrof_
image.set_shape(image_size + (3,))
images.append(image)
images_and_labels_list.append([images, label])
return images_and_labels_list

image_batch, label_batch = tf.train.batch_join(
images_and_labels_list, batch_size=batch_size_placeholder,
shapes=[image_size + (3,), ()], enqueue_many=True,
capacity=4 * nrof_preprocess_threads * 100,
allow_smaller_final_batch=True)

return image_batch, label_batch

def get_control_flag(control, field):
return tf.equal(tf.mod(tf.floor_div(control, field), 2), 1)
Expand Down Expand Up @@ -398,7 +361,7 @@ def split_dataset(dataset, split_ratio, min_nrof_images_per_class, mode):
raise ValueError('Invalid train/test split mode "%s"' % mode)
return train_set, test_set

def load_model(model):
def load_model(model, input_map=None):
# Check if the model is a model directory (containing a metagraph and a checkpoint file)
# or if it is a protobuf file with a frozen graph
model_exp = os.path.expanduser(model)
Expand All @@ -407,15 +370,15 @@ def load_model(model):
with gfile.FastGFile(model_exp,'rb') as f:
graph_def = tf.GraphDef()
graph_def.ParseFromString(f.read())
tf.import_graph_def(graph_def, name='')
tf.import_graph_def(graph_def, input_map=input_map, name='')
else:
print('Model directory: %s' % model_exp)
meta_file, ckpt_file = get_model_filenames(model_exp)

print('Metagraph file: %s' % meta_file)
print('Checkpoint file: %s' % ckpt_file)

saver = tf.train.import_meta_graph(os.path.join(model_exp, meta_file))
saver = tf.train.import_meta_graph(os.path.join(model_exp, meta_file), input_map=input_map)
saver.restore(tf.get_default_session(), os.path.join(model_exp, ckpt_file))

def get_model_filenames(model_dir):
Expand Down
9 changes: 1 addition & 8 deletions src/train_softmax.py
Original file line number Diff line number Diff line change
Expand Up @@ -119,21 +119,14 @@ def main(args):
labels_placeholder = tf.placeholder(tf.int32, shape=(None,1), name='labels')
control_placeholder = tf.placeholder(tf.int32, shape=(None,1), name='control')

images_and_labels = []
nrof_preprocess_threads = 4

input_queue = data_flow_ops.FIFOQueue(capacity=2000000,
dtypes=[tf.string, tf.int32, tf.int32],
shapes=[(1,), (1,), (1,)],
shared_name=None, name=None)
enqueue_op = input_queue.enqueue_many([image_paths_placeholder, labels_placeholder, control_placeholder], name='enqueue_op')
images_and_labels = facenet.create_input_pipeline(images_and_labels, input_queue, image_size, nrof_preprocess_threads)
image_batch, label_batch = facenet.create_input_pipeline(input_queue, image_size, nrof_preprocess_threads, batch_size_placeholder)

image_batch, label_batch = tf.train.batch_join(
images_and_labels, batch_size=batch_size_placeholder,
shapes=[image_size + (3,), ()], enqueue_many=True,
capacity=4 * nrof_preprocess_threads * args.batch_size,
allow_smaller_final_batch=True)
image_batch = tf.identity(image_batch, 'image_batch')
image_batch = tf.identity(image_batch, 'input')
label_batch = tf.identity(label_batch, 'label_batch')
Expand Down
123 changes: 88 additions & 35 deletions src/validate_on_lfw.py
Original file line number Diff line number Diff line change
Expand Up @@ -36,7 +36,7 @@
import lfw
import os
import sys
import math
from tensorflow.python.ops import data_flow_ops
from sklearn import metrics
from scipy.optimize import brentq
from scipy import interpolate
Expand All @@ -52,44 +52,89 @@ def main(args):

# Get the paths for the corresponding images
paths, actual_issame = lfw.get_paths(os.path.expanduser(args.lfw_dir), pairs)

# Load the model
facenet.load_model(args.model)

# Get input and output tensors
images_placeholder = tf.get_default_graph().get_tensor_by_name("input:0")
image_paths_placeholder = tf.placeholder(tf.string, shape=(None,1), name='image_paths')
labels_placeholder = tf.placeholder(tf.int32, shape=(None,1), name='labels')
batch_size_placeholder = tf.placeholder(tf.int32, name='batch_size')
control_placeholder = tf.placeholder(tf.int32, shape=(None,1), name='control')
phase_train_placeholder = tf.placeholder(tf.bool, name='phase_train')

nrof_preprocess_threads = 4
image_size = (args.image_size, args.image_size)
eval_input_queue = data_flow_ops.FIFOQueue(capacity=2000000,
dtypes=[tf.string, tf.int32, tf.int32],
shapes=[(1,), (1,), (1,)],
shared_name=None, name=None)
eval_enqueue_op = eval_input_queue.enqueue_many([image_paths_placeholder, labels_placeholder, control_placeholder], name='eval_enqueue_op')
image_batch, label_batch = facenet.create_input_pipeline(eval_input_queue, image_size, nrof_preprocess_threads, batch_size_placeholder)

# Load the model
input_map = {'image_batch': image_batch, 'label_batch': label_batch, 'phase_train': phase_train_placeholder}
facenet.load_model(args.model, input_map=input_map)

# Get output tensor
embeddings = tf.get_default_graph().get_tensor_by_name("embeddings:0")
phase_train_placeholder = tf.get_default_graph().get_tensor_by_name("phase_train:0")

#image_size = images_placeholder.get_shape()[1] # For some reason this doesn't work for frozen graphs
image_size = args.image_size
embedding_size = embeddings.get_shape()[1]

# Run forward pass to calculate embeddings
print('Runnning forward pass on LFW images')
batch_size = args.lfw_batch_size
nrof_images = len(paths)
nrof_batches = int(math.ceil(1.0*nrof_images / batch_size))
emb_array = np.zeros((nrof_images, embedding_size))
for i in range(nrof_batches):
start_index = i*batch_size
end_index = min((i+1)*batch_size, nrof_images)
paths_batch = paths[start_index:end_index]
images = facenet.load_data(paths_batch, False, False, image_size)
feed_dict = { images_placeholder:images, phase_train_placeholder:False }
emb_array[start_index:end_index,:] = sess.run(embeddings, feed_dict=feed_dict)

tpr, fpr, accuracy, val, val_std, far = lfw.evaluate(emb_array,
actual_issame, nrof_folds=args.lfw_nrof_folds)
#
coord = tf.train.Coordinator()
tf.train.start_queue_runners(coord=coord, sess=sess)

print('Accuracy: %1.3f+-%1.3f' % (np.mean(accuracy), np.std(accuracy)))
print('Validation rate: %2.5f+-%2.5f @ FAR=%2.5f' % (val, val_std, far))
evaluate(sess, eval_enqueue_op, image_paths_placeholder, labels_placeholder, phase_train_placeholder, batch_size_placeholder, control_placeholder,
embeddings, label_batch, paths, actual_issame, args.lfw_batch_size, args.lfw_nrof_folds, args.distance_metric, args.subtract_mean,
args.use_flipped_images, args.use_fixed_image_standardization)

auc = metrics.auc(fpr, tpr)
print('Area Under Curve (AUC): %1.3f' % auc)
eer = brentq(lambda x: 1. - x - interpolate.interp1d(fpr, tpr)(x), 0., 1.)
print('Equal Error Rate (EER): %1.3f' % eer)


def evaluate(sess, enqueue_op, image_paths_placeholder, labels_placeholder, phase_train_placeholder, batch_size_placeholder, control_placeholder,
embeddings, labels, image_paths, actual_issame, batch_size, nrof_folds, distance_metric, subtract_mean, use_flipped_images, use_fixed_image_standardization):
# Run forward pass to calculate embeddings
print('Runnning forward pass on LFW images')

# Enqueue one epoch of image paths and labels
nrof_embeddings = len(actual_issame)*2 # nrof_pairs * nrof_images_per_pair
nrof_flips = 2 if use_flipped_images else 1
nrof_images = nrof_embeddings * nrof_flips
labels_array = np.expand_dims(np.arange(0,nrof_images),1)
image_paths_array = np.expand_dims(np.repeat(np.array(image_paths),nrof_flips),1)
control_array = np.zeros_like(labels_array, np.int32)
if use_fixed_image_standardization:
control_array += np.ones_like(labels_array)*facenet.FIXED_STANDARDIZATION
if use_flipped_images:
# Flip every second image
control_array += (labels_array % 2)*facenet.FLIP
sess.run(enqueue_op, {image_paths_placeholder: image_paths_array, labels_placeholder: labels_array, control_placeholder: control_array})

embedding_size = int(embeddings.get_shape()[1])
assert nrof_images % batch_size == 0, 'The number of LFW images must be an integer multiple of the LFW batch size'
nrof_batches = nrof_images // batch_size
emb_array = np.zeros((nrof_images, embedding_size))
lab_array = np.zeros((nrof_images,))
for i in range(nrof_batches):
feed_dict = {phase_train_placeholder:False, batch_size_placeholder:batch_size}
emb, lab = sess.run([embeddings, labels], feed_dict=feed_dict)
lab_array[lab] = lab
emb_array[lab, :] = emb
if i % 10 == 9:
print('.', end='')
sys.stdout.flush()
print('')
embeddings = np.zeros((nrof_embeddings, embedding_size*nrof_flips))
if use_flipped_images:
# Concatenate embeddings for flipped and non flipped iversion of the images
embeddings[:,:embedding_size] = emb_array[0::2,:]
embeddings[:,embedding_size:] = emb_array[1::2,:]
else:
embeddings = emb_array

assert np.array_equal(lab_array, np.arange(nrof_images))==True, 'Wrong labels used for evaluation, possibly caused by training examples left in the input pipeline'
tpr, fpr, accuracy, val, val_std, far = lfw.evaluate(embeddings, actual_issame, nrof_folds=nrof_folds, distance_metric=distance_metric, subtract_mean=subtract_mean)

print('Accuracy: %2.5f+-%2.5f' % (np.mean(accuracy), np.std(accuracy)))
print('Validation rate: %2.5f+-%2.5f @ FAR=%2.5f' % (val, val_std, far))

auc = metrics.auc(fpr, tpr)
print('Area Under Curve (AUC): %1.3f' % auc)
eer = brentq(lambda x: 1. - x - interpolate.interp1d(fpr, tpr)(x), 0., 1.)
print('Equal Error Rate (EER): %1.3f' % eer)

def parse_arguments(argv):
parser = argparse.ArgumentParser()

Expand All @@ -105,6 +150,14 @@ def parse_arguments(argv):
help='The file containing the pairs to use for validation.', default='data/pairs.txt')
parser.add_argument('--lfw_nrof_folds', type=int,
help='Number of folds to use for cross validation. Mainly used for testing.', default=10)
parser.add_argument('--distance_metric', type=int,
help='Distance metric 0:euclidian, 1:cosine similarity.', default=0)
parser.add_argument('--use_flipped_images',
help='Concatenates embeddings for the image and its horizontally flipped counterpart.', action='store_true')
parser.add_argument('--subtract_mean',
help='Subtract feature mean before calculating distance.', action='store_true')
parser.add_argument('--use_fixed_image_standardization',
help='Performs fixed standardization of images.', action='store_true')
return parser.parse_args(argv)

if __name__ == '__main__':
Expand Down
Loading

0 comments on commit 9852362

Please sign in to comment.